Skip to content

Commit 11c7f66

Browse files
authored
[Simple GLA] Add comments & Fix bad grad
1 parent fd44b54 commit 11c7f66

File tree

1 file changed

+71
-95
lines changed

1 file changed

+71
-95
lines changed

fla/ops/simple_gla/chunk.py

Lines changed: 71 additions & 95 deletions
Original file line numberDiff line numberDiff line change
@@ -1,7 +1,7 @@
11
# -*- coding: utf-8 -*-
22
# Copyright (c) 2023, Yu Zhang, Songlin Yang
33

4-
from typing import Tuple
4+
from typing import Optional, Tuple
55

66
import torch
77
import triton
@@ -11,23 +11,14 @@
1111
from fla.utils import contiguous
1212

1313

14-
@torch.jit.script
15-
def normalize_output(q, k, o):
16-
k = k.transpose(-2, -1)
17-
k = k.cumsum(-1)
18-
k = k.transpose(-2, -1)
19-
z = (q * k).sum(-1, keepdim=True)
20-
return o / (z + 1e-5)
21-
22-
2314
@triton.jit
2415
def chunk_simple_gla_fwd_kernel_h(
2516
k,
2617
v,
2718
h,
2819
g,
29-
initial_state, # initial state of the chunk [B, H, D_head_K, D_head_V]
30-
final_state, # final state of the chunk [B, H, D_head_K, D_head_V]
20+
h0,
21+
ht,
3122
s_qk_h,
3223
s_qk_t,
3324
s_qk_d,
@@ -36,7 +27,6 @@ def chunk_simple_gla_fwd_kernel_h(
3627
s_vo_d,
3728
s_h_h,
3829
s_h_t,
39-
H: tl.constexpr,
4030
T: tl.constexpr,
4131
K: tl.constexpr,
4232
V: tl.constexpr,
@@ -53,17 +43,13 @@ def chunk_simple_gla_fwd_kernel_h(
5343
b_h = tl.zeros([BK, BV], dtype=tl.float32)
5444

5545
if USE_INITIAL_STATE:
56-
p_h0 = tl.make_block_ptr(initial_state + i_bh * K * V,
57-
(K, V), (V, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
46+
p_h0 = tl.make_block_ptr(h0 + i_bh * K * V, (K, V), (V, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
5847
b_h = tl.load(p_h0, boundary_check=(0, 1)).to(tl.float32)
5948

6049
for i_t in range(NT):
61-
p_k = tl.make_block_ptr(
62-
k + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
63-
p_v = tl.make_block_ptr(
64-
v + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
65-
p_h = tl.make_block_ptr(h + i_bh * s_h_h + i_t * K * V,
66-
(K, V), (s_h_t, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
50+
p_k = tl.make_block_ptr(k + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
51+
p_v = tl.make_block_ptr(v + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
52+
p_h = tl.make_block_ptr(h + i_bh * s_h_h + i_t * K * V, (K, V), (s_h_t, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
6753

6854
tl.store(p_h, b_h.to(p_h.dtype.element_ty), boundary_check=(0, 1))
6955
# [BK, BT]
@@ -72,13 +58,12 @@ def chunk_simple_gla_fwd_kernel_h(
7258
b_v = tl.load(p_v, boundary_check=(0, 1))
7359
# [BK, BV]
7460
b_g_last = tl.load(g + i_bh * T + i_t * BT + BT - 1)
75-
b_h *= tl.math.exp2(b_g_last)
61+
b_h *= tl.exp(b_g_last)
7662
b_g = tl.load(g + i_bh * T + i_t * BT + tl.arange(0, BT))
77-
b_h += tl.dot(b_k, (b_v * tl.math.exp2(b_g_last - b_g)[:, None]).to(b_k.dtype), allow_tf32=False)
63+
b_h += tl.dot(b_k, (b_v * tl.exp(b_g_last - b_g)[:, None]).to(b_k.dtype), allow_tf32=False)
7864

7965
if STORE_FINAL_STATE:
80-
p_ht = tl.make_block_ptr(
81-
final_state + i_bh * K * V, (K, V), (V, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
66+
p_ht = tl.make_block_ptr(ht + i_bh * K * V, (K, V), (V, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
8267
tl.store(p_ht, b_h.to(p_ht.dtype.element_ty), boundary_check=(0, 1))
8368

8469

@@ -99,7 +84,6 @@ def chunk_simple_gla_fwd_kernel_o(
9984
s_h_h,
10085
s_h_t,
10186
scale,
102-
H: tl.constexpr,
10387
T: tl.constexpr,
10488
K: tl.constexpr,
10589
V: tl.constexpr,
@@ -115,12 +99,9 @@ def chunk_simple_gla_fwd_kernel_o(
11599
b_o = tl.zeros([BT, BV], dtype=tl.float32)
116100
b_s = tl.zeros([BT, BT], dtype=tl.float32)
117101
for i_k in range(tl.cdiv(K, BK)):
118-
p_q = tl.make_block_ptr(
119-
q + i_bh * s_qk_h, (T, K), (s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
120-
p_k = tl.make_block_ptr(
121-
k + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
122-
p_h = tl.make_block_ptr(h + i_bh * s_h_h + i_t * K * V,
123-
(K, V), (s_h_t, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
102+
p_q = tl.make_block_ptr(q + i_bh * s_qk_h, (T, K), (s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
103+
p_k = tl.make_block_ptr(k + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
104+
p_h = tl.make_block_ptr(h + i_bh * s_h_h + i_t * K * V, (K, V), (s_h_t, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
124105

125106
# [BT, BK]
126107
b_q = tl.load(p_q, boundary_check=(0, 1))
@@ -135,16 +116,14 @@ def chunk_simple_gla_fwd_kernel_o(
135116

136117
p_g = g + i_bh * T + i_t * BT + tl.arange(0, BT)
137118
b_g = tl.load(p_g)
138-
b_o = b_o * tl.math.exp2(b_g)[:, None]
139-
b_s = b_s * tl.math.exp2(b_g[:, None] - b_g[None, :])
119+
b_o = b_o * tl.exp(b_g)[:, None]
120+
b_s = b_s * tl.exp(b_g[:, None] - b_g[None, :])
140121
b_s = tl.where(m_s, b_s, 0)
141122

142-
p_v = tl.make_block_ptr(v + i_bh * s_vo_h, (T, V),
143-
(s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
123+
p_v = tl.make_block_ptr(v + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
144124
b_v = tl.load(p_v, boundary_check=(0, 1))
145125
b_o = (b_o + tl.dot(b_s.to(b_v.dtype), b_v, allow_tf32=False)) * scale
146-
p_o = tl.make_block_ptr(o + i_bh * s_vo_h, (T, V),
147-
(s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
126+
p_o = tl.make_block_ptr(o + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
148127
tl.store(p_o, b_o.to(p_o.dtype.element_ty), boundary_check=(0, 1))
149128

150129

@@ -163,7 +142,6 @@ def chunk_simple_gla_bwd_kernel_dh(
163142
s_h_h,
164143
s_h_t,
165144
scale,
166-
H: tl.constexpr,
167145
T: tl.constexpr,
168146
K: tl.constexpr,
169147
V: tl.constexpr,
@@ -177,22 +155,18 @@ def chunk_simple_gla_bwd_kernel_dh(
177155
# [BK, BV]
178156
b_dh = tl.zeros([BK, BV], dtype=tl.float32)
179157
for i_t in range(NT - 1, -1, -1):
180-
p_q = tl.make_block_ptr(
181-
q + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
182-
p_do = tl.make_block_ptr(
183-
do + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
184-
p_dh = tl.make_block_ptr(dh + i_bh * s_h_h + i_t * K * V,
185-
(K, V), (s_h_t, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
158+
p_q = tl.make_block_ptr(q + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
159+
p_do = tl.make_block_ptr(do + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
160+
p_dh = tl.make_block_ptr(dh + i_bh * s_h_h + i_t * K * V, (K, V), (s_h_t, 1), (i_k * BK, i_v * BV), (BK, BV), (1, 0))
186161

187162
tl.store(p_dh, b_dh.to(p_dh.dtype.element_ty), boundary_check=(0, 1))
188163
# [BK, BT]
189164
b_q = tl.load(p_q, boundary_check=(0, 1))
190-
b_q = (b_q * scale * tl.math.exp2(tl.load(g + i_bh * T +
191-
i_t * BT + tl.arange(0, BT)))[None, :]).to(b_q.dtype)
165+
b_q = (b_q * scale * tl.exp(tl.load(g + i_bh * T + i_t * BT + tl.arange(0, BT)))[None, :]).to(b_q.dtype)
192166
# [BT, V]
193167
b_do = tl.load(p_do, boundary_check=(0, 1))
194168
# [BK, BV]
195-
b_dh *= tl.math.exp2(tl.load(g + i_bh * T + i_t * BT + BT - 1))
169+
b_dh *= tl.exp(tl.load(g + i_bh * T + i_t * BT + BT - 1))
196170
b_dh += tl.dot(b_q, b_do.to(b_q.dtype), allow_tf32=False)
197171

198172

@@ -217,8 +191,6 @@ def chunk_simple_gla_bwd_kernel_dqkv(
217191
s_h_h,
218192
s_h_t,
219193
scale,
220-
B: tl.constexpr,
221-
H: tl.constexpr,
222194
T: tl.constexpr,
223195
K: tl.constexpr,
224196
V: tl.constexpr,
@@ -231,35 +203,28 @@ def chunk_simple_gla_bwd_kernel_dqkv(
231203
n_bh = tl.num_programs(2)
232204
o_i = tl.arange(0, BT)
233205

234-
p_q = tl.make_block_ptr(q + i_bh * s_qk_h, (K, T),
235-
(s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
236-
p_k = tl.make_block_ptr(k + i_bh * s_qk_h, (T, K),
237-
(s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
206+
p_q = tl.make_block_ptr(q + i_bh * s_qk_h, (K, T), (s_qk_d, s_qk_t), (i_k * BK, i_t * BT), (BK, BT), (0, 1))
207+
p_k = tl.make_block_ptr(k + i_bh * s_qk_h, (T, K), (s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
238208

239209
b_q = tl.load(p_q, boundary_check=(0, 1))
240210
b_k = tl.load(p_k, boundary_check=(0, 1))
241211
b_s = tl.dot(b_k, b_q, allow_tf32=False)
242212
p_g = g + i_bh * T + i_t * BT + tl.arange(0, BT)
243213
b_g = tl.load(p_g)
244214
b_g_last = tl.load(g + i_bh * T + i_t * BT + BT - 1)
245-
mask = tl.math.exp2(b_g[None, :] - b_g[:, None])
215+
mask = tl.exp(b_g[None, :] - b_g[:, None])
246216
mask = tl.where(o_i[:, None] <= o_i[None, :], mask * scale, 0)
247217
b_s = b_s * mask
248218

249219
b_dq = tl.zeros([BT, BK], dtype=tl.float32)
250220
b_dk = tl.zeros([BT, BK], dtype=tl.float32)
251221
b_ds = tl.zeros([BT, BT], dtype=tl.float32)
252222
for i_v in range(tl.cdiv(V, BV)):
253-
p_v = tl.make_block_ptr(
254-
v + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
255-
p_h = tl.make_block_ptr(h + i_bh * s_h_h, (V, NT * K), (1, s_h_t),
256-
(i_v * BV, i_t * K + i_k * BK), (BV, BK), (0, 1))
257-
p_do = tl.make_block_ptr(
258-
do + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
259-
p_dh = tl.make_block_ptr(dh + i_bh * s_h_h, (NT * K, V),
260-
(s_h_t, 1), (i_t * K + i_k * BK, i_v * BV), (BK, BV), (1, 0))
261-
p_dv = tl.make_block_ptr(dv + (i_k*n_bh+i_bh)*s_vo_h, (T, V),
262-
(s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
223+
p_v = tl.make_block_ptr(v + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
224+
p_h = tl.make_block_ptr(h + i_bh * s_h_h, (V, NT * K), (1, s_h_t), (i_v * BV, i_t * K + i_k * BK), (BV, BK), (0, 1))
225+
p_do = tl.make_block_ptr(do + i_bh * s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
226+
p_dh = tl.make_block_ptr(dh + i_bh * s_h_h, (NT * K, V), (s_h_t, 1), (i_t * K + i_k * BK, i_v * BV), (BK, BV), (1, 0))
227+
p_dv = tl.make_block_ptr(dv + (i_k*n_bh+i_bh)*s_vo_h, (T, V), (s_vo_t, s_vo_d), (i_t * BT, i_v * BV), (BT, BV), (1, 0))
263228
# [BT, BV]
264229
b_v = tl.load(p_v, boundary_check=(0, 1))
265230
b_do = tl.load(p_do, boundary_check=(0, 1))
@@ -273,21 +238,19 @@ def chunk_simple_gla_bwd_kernel_dqkv(
273238
b_dq += tl.dot(b_do, b_h, allow_tf32=False) * scale
274239
b_dk += tl.dot(b_v, tl.trans(b_dh), allow_tf32=False)
275240
# [BT, BV]
276-
b_dv = tl.dot(b_k, b_dh, allow_tf32=False) * tl.math.exp2(-b_g + b_g_last)[:, None] + \
277-
tl.dot(b_s.to(b_q.dtype), b_do, allow_tf32=False)
241+
b_dv = tl.dot(b_k, b_dh, allow_tf32=False) * tl.exp(-b_g + b_g_last)[:, None]
242+
b_dv += tl.dot(b_s.to(b_q.dtype), b_do, allow_tf32=False)
278243
tl.store(p_dv, b_dv.to(p_dv.dtype.element_ty), boundary_check=(0, 1))
279244

280-
b_dq = b_dq * tl.math.exp2(b_g)[:, None]
281-
b_dk = b_dk * tl.math.exp2(-b_g + b_g_last)[:, None]
245+
b_dq = b_dq * tl.exp(b_g)[:, None]
246+
b_dk = b_dk * tl.exp(-b_g + b_g_last)[:, None]
282247
b_ds = b_ds * tl.trans(mask)
283248
b_ds = b_ds.to(b_k.dtype)
284249
# [BT, BK]
285250
b_dq += tl.dot(b_ds, b_k, allow_tf32=False)
286251
b_dk += tl.trans(tl.dot(b_q, b_ds, allow_tf32=False))
287-
p_dq = tl.make_block_ptr(dq + i_bh * s_qk_h, (T, K),
288-
(s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
289-
p_dk = tl.make_block_ptr(dk + i_bh * s_qk_h, (T, K),
290-
(s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
252+
p_dq = tl.make_block_ptr(dq + i_bh * s_qk_h, (T, K), (s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
253+
p_dk = tl.make_block_ptr(dk + i_bh * s_qk_h, (T, K), (s_qk_t, s_qk_d), (i_t * BT, i_k * BK), (BT, BK), (1, 0))
291254
tl.store(p_dq, b_dq.to(p_dq.dtype.element_ty), boundary_check=(0, 1))
292255
tl.store(p_dk, b_dk.to(p_dk.dtype.element_ty), boundary_check=(0, 1))
293256

@@ -300,19 +263,14 @@ class SimpleGLAFunction(torch.autograd.Function):
300263
def forward(ctx, q, k, v, g, scale, initial_state, output_final_state):
301264
B, H, T, K, V = *q.shape, v.shape[-1]
302265
BT = 64
303-
BK, BV = min(64, triton.next_power_of_2(K)), min(
304-
64, triton.next_power_of_2(V))
266+
BK, BV = min(64, triton.next_power_of_2(K)), min(64, triton.next_power_of_2(V))
305267
NT, NK, NV = triton.cdiv(T, BT), triton.cdiv(K, BK), triton.cdiv(V, BV)
306-
num_stages = 1
307268
num_warps = 4 if BK == 64 else 2
269+
num_stages = 1
308270

309-
if scale is None:
310-
scale = K ** -0.5
311-
312-
BT = 64
313271
assert T % BT == 0, 'sequence length must be divisible by BT'
314272
g = g.reshape(B, H, -1, BT)
315-
g = g.cumsum(-1) * 1.44269504
273+
g = g.cumsum(-1)
316274
g = g.reshape(B, H, -1)
317275

318276
final_state = None
@@ -326,7 +284,7 @@ def forward(ctx, q, k, v, g, scale, initial_state, output_final_state):
326284
q.stride(1), q.stride(2), q.stride(3),
327285
v.stride(1), v.stride(2), v.stride(3),
328286
h.stride(1), h.stride(2),
329-
H=H, T=T, K=K, V=V, BT=BT, BK=BK, BV=BV, NT=NT,
287+
T=T, K=K, V=V, BT=BT, BK=BK, BV=BV, NT=NT,
330288
USE_INITIAL_STATE=initial_state is not None,
331289
STORE_FINAL_STATE=output_final_state,
332290
num_warps=num_warps,
@@ -340,30 +298,29 @@ def forward(ctx, q, k, v, g, scale, initial_state, output_final_state):
340298
v.stride(1), v.stride(2), v.stride(3),
341299
h.stride(1), h.stride(2),
342300
scale,
343-
H=H, T=T, K=K, V=V, BT=BT, BK=BK, BV=BV,
301+
T=T, K=K, V=V, BT=BT, BK=BK, BV=BV,
344302
num_warps=num_warps,
345303
num_stages=num_stages
346304
)
347305

348306
ctx.save_for_backward(q, k, v, h, g)
307+
ctx.scale = scale
349308
return o.to(q.dtype), final_state
350309

351310
@staticmethod
352311
@custom_bwd
353312
@contiguous
354-
def backward(ctx, do, scale, d_ht=None):
313+
def backward(ctx, do, dht=None):
355314
q, k, v, h, g = ctx.saved_tensors
356315

357316
B, H, T, K, V = *q.shape, v.shape[-1]
358317
BT = 64
359-
BK, BV = min(32 if q.dtype == torch.float32 else 64, triton.next_power_of_2(K)), min(
360-
32 if q.dtype == torch.float32 else 64, triton.next_power_of_2(V))
318+
BK = min(32 if q.dtype == torch.float32 else 64, triton.next_power_of_2(K))
319+
BV = min(32 if q.dtype == torch.float32 else 64, triton.next_power_of_2(V))
361320
NT, NK, NV = triton.cdiv(T, BT), triton.cdiv(K, BK), triton.cdiv(V, BV)
362-
num_stages = 1
363321
num_warps = 4 if BK == 64 else 2
364-
365-
if scale is None:
366-
scale = K ** -0.5
322+
num_stages = 1
323+
scale = ctx.scale
367324

368325
dh = q.new_empty(B, H, NT * K, V)
369326
grid = (NK, NV, B * H)
@@ -373,7 +330,7 @@ def backward(ctx, do, scale, d_ht=None):
373330
v.stride(1), v.stride(2), v.stride(3),
374331
dh.stride(1), dh.stride(2),
375332
scale,
376-
H=H, T=T, K=K, V=V, BT=BT, BK=BK, BV=BV, NT=NT,
333+
T=T, K=K, V=V, BT=BT, BK=BK, BV=BV, NT=NT,
377334
num_warps=num_warps,
378335
num_stages=num_stages
379336
)
@@ -389,7 +346,7 @@ def backward(ctx, do, scale, d_ht=None):
389346
v.stride(1), v.stride(2), v.stride(3),
390347
dh.stride(1), dh.stride(2),
391348
scale,
392-
B=B, H=H, T=T, K=K, V=V, BT=BT, BK=BK, BV=BV, NT=NT,
349+
T=T, K=K, V=V, BT=BT, BK=BK, BV=BV, NT=NT,
393350
num_warps=num_warps,
394351
num_stages=num_stages
395352
)
@@ -409,12 +366,31 @@ def chunk_simple_gla(
409366
k: torch.Tensor,
410367
v: torch.Tensor,
411368
g: torch.Tensor, # log decay
412-
scale: float = None,
369+
scale: Optional[float] = None,
413370
initial_state: torch.Tensor = None,
414371
output_final_state: bool = False
415372
) -> Tuple[torch.Tensor, torch.Tensor]:
416-
if initial_state is not None:
417-
initial_state = initial_state.detach()
373+
r"""
374+
Args:
375+
q (torch.Tensor):
376+
queries of shape `(B, H, T, K)`
377+
k (torch.Tensor):
378+
keys of shape `(B, H, T, K)`
379+
v (torch.Tensor):
380+
values of shape `(B, H, T, V)`
381+
g (torch.Tensor):
382+
Forget gates of shape `(B, H, T)` applied to keys.
383+
Compared to GLA, the gating is head-wise instead of elementwise.
384+
scale (Optional[int]):
385+
Scale factor for the attention scores.
386+
If not provided, it will default to `1 / sqrt(K)`. Default: `None`.
387+
initial_state (Optional[torch.Tensor]):
388+
Initial state of shape `(B, H, K, V)`. Default: `None`.
389+
output_final_state (Optional[bool]):
390+
Whether to output the final state of shape `(B, H, K, V)`. Default: `False`.
391+
"""
392+
if scale is None:
393+
scale = k.shape[-1] ** -0.5
418394
g = g.float()
419395
o, final_state = SimpleGLAFunction.apply(q, k, v, g, scale, initial_state, output_final_state)
420396
return o, final_state

0 commit comments

Comments
 (0)