Skip to content

Commit 9f4ab2e

Browse files
authored
[Inference] Adapt to Fused rotary (#5348)
* revise rotary embedding * remove useless print * adapt * fix * add * fix * modeling * fix * fix * fix
1 parent 35382a7 commit 9f4ab2e

File tree

5 files changed

+161
-22
lines changed

5 files changed

+161
-22
lines changed

colossalai/inference/modeling/models/nopadding_llama.py

+2-3
Original file line numberDiff line numberDiff line change
@@ -282,11 +282,10 @@ def forward(
282282
torch.bmm(hidden_states, self.qkv_weight).view(3, token_nums, self.num_heads, self.head_dim).unbind(0)
283283
)
284284

285-
rotary_embedding(query_states, key_states, cos_sin[0], cos_sin[1])
286-
287285
block_size = k_cache.size(-2)
288286

289287
if is_prompts:
288+
rotary_embedding(query_states, key_states, cos_sin[0], cos_sin[1])
290289
attn_output = context_attention_unpadded(
291290
q=query_states,
292291
k=key_states,
@@ -301,7 +300,7 @@ def forward(
301300
sm_scale=sm_scale,
302301
)
303302
else:
304-
copy_kv_to_blocked_cache(key_states, k_cache, kv_lengths=sequence_lengths, block_tables=block_tables)
303+
rotary_embedding(query_states, key_states, cos_sin[0], cos_sin[1], k_cache, block_tables, sequence_lengths)
305304
copy_kv_to_blocked_cache(value_states, v_cache, kv_lengths=sequence_lengths, block_tables=block_tables)
306305
attn_output = flash_decoding_attention(
307306
q=query_states,

colossalai/kernel/triton/kvcache_copy.py

-1
Original file line numberDiff line numberDiff line change
@@ -75,7 +75,6 @@ def copy_kv_to_blocked_cache(
7575
block_size = k_cache.size(-2)
7676

7777
num_warps = 8 if head_dim > 128 else 4
78-
7978
grid = (bsz, num_kv_heads)
8079
_copy_to_kvcache_seqlen1_kernel[grid](
8180
k,

colossalai/kernel/triton/no_pad_rotary_embedding.py

+128-8
Original file line numberDiff line numberDiff line change
@@ -222,11 +222,11 @@ def fused_rotary_embedding_kernel(
222222
out_k0 = loaded_k0 * loaded_cos[:, None, :] - loaded_k1 * loaded_sin[:, None, :]
223223
out_k1 = loaded_k0 * loaded_sin[:, None, :] + loaded_k1 * loaded_cos[:, None, :] # total_tokens, head_num, head_dim
224224

225-
past_kv_seq_len = tl.load(context_lengths + tokens_range) - 1
225+
past_kv_seq_len = tl.load(context_lengths + tokens_range, mask=(tokens_range < q_total_tokens)) - 1
226226

227227
last_block_idx = past_kv_seq_len // block_size
228228
block_table_ptr = BLOCK_TABLES + tokens_range * bts_stride
229-
block_ids = tl.load(block_table_ptr + last_block_idx * btb_stride)
229+
block_ids = tl.load(block_table_ptr + last_block_idx * btb_stride, mask=(tokens_range < q_total_tokens))
230230
offsets_in_last_block = (past_kv_seq_len % block_size) * cachebs_stride
231231

232232
kv_range0 = (
@@ -274,6 +274,122 @@ def fused_rotary_embedding_kernel(
274274
)
275275

276276

277+
@triton.jit
278+
def fused_rotary_embedding_kernel_v2(
279+
q,
280+
k,
281+
cos,
282+
sin,
283+
kv_cache,
284+
BLOCK_TABLES,
285+
context_lengths,
286+
q_token_stride,
287+
q_head_stride,
288+
k_token_stride,
289+
k_head_stride,
290+
head_dim_stride,
291+
cos_token_stride,
292+
cos_stride,
293+
cacheb_stride,
294+
cacheh_stride,
295+
cachebs_stride,
296+
cached_stride,
297+
bts_stride,
298+
btb_stride,
299+
block_size,
300+
q_total_tokens,
301+
Q_HEAD_NUM: tl.constexpr,
302+
K_HEAD_NUM: tl.constexpr,
303+
HEAD_DIM: tl.constexpr,
304+
):
305+
block_head_index = tl.program_id(0)
306+
if block_head_index >= Q_HEAD_NUM:
307+
return
308+
block_token_index = tl.program_id(1)
309+
310+
dim_range0 = tl.arange(0, HEAD_DIM // 2)
311+
dim_range1 = tl.arange(HEAD_DIM // 2, HEAD_DIM)
312+
313+
off_q0 = block_token_index * q_token_stride + block_head_index * q_head_stride + dim_range0 * head_dim_stride
314+
off_q1 = block_token_index * q_token_stride + block_head_index * q_head_stride + dim_range1 * head_dim_stride
315+
off_k0 = block_token_index * k_token_stride + block_head_index * k_head_stride + dim_range0 * head_dim_stride
316+
off_k1 = block_token_index * k_token_stride + block_head_index * k_head_stride + dim_range1 * head_dim_stride
317+
318+
loaded_q0 = tl.load(
319+
q + off_q0,
320+
)
321+
loaded_q1 = tl.load(
322+
q + off_q1,
323+
)
324+
325+
loaded_k0 = tl.load(
326+
k + off_k0,
327+
)
328+
329+
loaded_k1 = tl.load(
330+
k + off_k1,
331+
)
332+
333+
off_cos_sin = block_token_index * cos_token_stride + dim_range0 * cos_stride
334+
335+
loaded_cos = tl.load(cos + off_cos_sin, mask=(block_token_index < q_total_tokens), other=0.0)
336+
loaded_sin = tl.load(sin + off_cos_sin, mask=(block_token_index < q_total_tokens), other=0.0)
337+
338+
out_q0 = loaded_q0 * loaded_cos - loaded_q1 * loaded_sin
339+
out_q1 = loaded_q0 * loaded_sin + loaded_q1 * loaded_cos
340+
341+
out_k0 = loaded_k0 * loaded_cos - loaded_k1 * loaded_sin
342+
out_k1 = loaded_k0 * loaded_sin + loaded_k1 * loaded_cos # total_tokens, head_num, head_dim
343+
344+
past_kv_seq_len = tl.load(context_lengths + block_token_index) - 1
345+
346+
last_block_idx = past_kv_seq_len // block_size
347+
block_table_ptr = BLOCK_TABLES + block_token_index * bts_stride
348+
block_ids = tl.load(block_table_ptr + last_block_idx * btb_stride, mask=(block_token_index < q_total_tokens))
349+
offsets_in_last_block = (past_kv_seq_len % block_size) * cachebs_stride
350+
351+
kv_range0 = (
352+
block_ids * cacheb_stride
353+
+ block_head_index * cacheh_stride
354+
+ offsets_in_last_block
355+
+ dim_range0 * cached_stride
356+
)
357+
kv_range1 = (
358+
block_ids * cacheb_stride
359+
+ block_head_index * cacheh_stride
360+
+ offsets_in_last_block
361+
+ dim_range1 * cached_stride
362+
)
363+
364+
tl.store(
365+
kv_cache + kv_range0,
366+
out_k0,
367+
)
368+
tl.store(
369+
kv_cache + kv_range1,
370+
out_k1,
371+
)
372+
373+
# concat
374+
tl.store(
375+
q + off_q0,
376+
out_q0,
377+
)
378+
tl.store(
379+
q + off_q1,
380+
out_q1,
381+
)
382+
tl.store(
383+
k + off_k0,
384+
out_k0,
385+
)
386+
tl.store(
387+
k + off_k1,
388+
out_k1,
389+
)
390+
391+
392+
@torch.no_grad()
277393
def rotary_embedding(
278394
q: torch.Tensor,
279395
k: torch.Tensor,
@@ -297,12 +413,13 @@ def rotary_embedding(
297413
assert q.size(0) == k.size(0)
298414
BLOCK_HEAD = 4
299415
BLOCK_TOKENS = 4
300-
grid = lambda META: (triton.cdiv(q_head_num, META["BLOCK_HEAD"]), triton.cdiv(q_total_tokens, META["BLOCK_TOKENS"]))
301416

302-
if head_dim >= 256:
417+
if head_dim >= 1024:
303418
num_warps = 32
304-
elif head_dim >= 128:
419+
elif head_dim >= 512:
305420
num_warps = 16
421+
elif head_dim >= 256:
422+
num_warps = 8
306423
else:
307424
num_warps = 4
308425

@@ -318,6 +435,10 @@ def rotary_embedding(
318435
cos_token_stride = cos.stride(0)
319436
cos_stride = cos.stride(1)
320437
if k_cache == None:
438+
grid = lambda META: (
439+
triton.cdiv(q_head_num, META["BLOCK_HEAD"]),
440+
triton.cdiv(q_total_tokens, META["BLOCK_TOKENS"]),
441+
)
321442
rotary_embedding_kernel[grid](
322443
q,
323444
k,
@@ -339,7 +460,8 @@ def rotary_embedding(
339460
num_warps=num_warps,
340461
)
341462
else:
342-
fused_rotary_embedding_kernel[grid](
463+
grid = (triton.next_power_of_2(q_head_num), q_total_tokens)
464+
fused_rotary_embedding_kernel_v2[grid](
343465
q,
344466
k,
345467
cos,
@@ -365,8 +487,6 @@ def rotary_embedding(
365487
Q_HEAD_NUM=q_head_num,
366488
K_HEAD_NUM=k_head_num,
367489
HEAD_DIM=head_dim,
368-
BLOCK_HEAD=BLOCK_HEAD,
369-
BLOCK_TOKENS=BLOCK_TOKENS,
370490
num_warps=num_warps,
371491
)
372492
return

examples/inference/run_benchmark.sh

+1
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,5 @@
11
ROOT=$(realpath $(dirname $0))
2+
echo $ROOT
23
PY_SCRIPT=${ROOT}/benchmark_llama.py
34
GPU=$(nvidia-smi -L | head -1 | cut -d' ' -f4 | cut -d'-' -f1)
45
mode=$1

tests/test_infer/test_ops/triton/test_rotary_embdding_unpad.py

+30-10
Original file line numberDiff line numberDiff line change
@@ -3,7 +3,7 @@
33
from packaging import version
44
from transformers.models.llama.modeling_llama import LlamaRotaryEmbedding, apply_rotary_pos_emb
55

6-
from colossalai.kernel.triton import rotary_embedding
6+
from colossalai.kernel.triton import copy_kv_to_blocked_cache, rotary_embedding
77
from tests.test_infer.test_ops.triton.kernel_utils import mock_alloc_block_table_and_kvcache_v2
88

99
try:
@@ -94,8 +94,8 @@ def test_rotary_emb(BATCH_SIZE, SEQ_LEN, H, D, dtype):
9494
x_names=["num_tokens"],
9595
x_vals=[2**i for i in range(4, 11)],
9696
line_arg="provider",
97-
line_vals=["torch_rotary_emb_func", "triton_rotary_emb_func"],
98-
line_names=["torch_rotary_emb_func", "triton_rotary_emb_func"],
97+
line_vals=["no_fused_rotary_emb_func", "fused_triton_rotary_emb_func"],
98+
line_names=["no_fused_rotary_emb_func", "fused_triton_rotary_emb_func"],
9999
styles=[("red", "-"), ("blue", "-")],
100100
ylabel="ms",
101101
plot_name=f"rotary_emb-batch-{BATCH}",
@@ -110,23 +110,43 @@ def benchmark_rotary_emb(
110110
num_tokens: int,
111111
num_kv_heads: int,
112112
):
113+
BATCH_SIZE = 4
114+
SEQ_LEN = num_tokens // BATCH_SIZE
115+
max_num_blocks_per_seq = 8
116+
block_size = 64
113117
warmup = 10
114118
rep = 100
115119

116-
head_dim = 128
120+
head_dim = 256
117121
dtype = torch.float16
122+
118123
q_shape = (num_tokens, num_kv_heads, head_dim)
119124
q = -2.3 + 0.5 * torch.randn(q_shape, dtype=dtype, device="cuda")
120125
k_shape = (num_tokens, num_kv_heads, head_dim)
121126
k = -2.3 + 0.5 * torch.randn(k_shape, dtype=dtype, device="cuda")
122127
cos_shape = (num_tokens, head_dim // 2)
123128
cos = -1.2 + 0.5 * torch.randn(cos_shape, dtype=dtype, device="cuda")
124129
sin = -2.0 + 0.5 * torch.randn(cos_shape, dtype=dtype, device="cuda")
130+
cache_shape = (BATCH_SIZE * max_num_blocks_per_seq, num_kv_heads, block_size, head_dim)
131+
k_cache = torch.zeros(size=cache_shape, dtype=dtype, device="cuda")
132+
v = torch.randn_like(k)
133+
v_cache = torch.zeros_like(k_cache)
134+
past_kv_seq_lengths = torch.tensor([SEQ_LEN - 1 for _ in range(BATCH_SIZE)], dtype=torch.int32, device="cuda")
135+
block_tables = mock_alloc_block_table_and_kvcache_v2(
136+
k, v, k_cache, v_cache, past_kv_seq_lengths, BATCH_SIZE, max_num_blocks_per_seq, block_size
137+
)
138+
new_k = torch.randn((BATCH_SIZE, num_kv_heads, head_dim), dtype=dtype, device="cuda")
139+
new_q = torch.randn_like(new_k)
140+
kv_seq_lengths = past_kv_seq_lengths + 1
141+
block_tables = block_tables.to(device="cuda")
125142

126-
if provider == "torch_rotary_emb_func":
127-
fn = lambda: torch_rotary_emb(q, cos, sin)
128-
elif provider == "triton_rotary_emb_func":
129-
fn = lambda: rotary_embedding(q, k, cos, sin)
143+
if provider == "no_fused_rotary_emb_func":
144+
fn = lambda: [
145+
rotary_embedding(new_q, new_k, cos, sin),
146+
copy_kv_to_blocked_cache(new_k, k_cache, kv_lengths=kv_seq_lengths, block_tables=block_tables),
147+
]
148+
elif provider == "fused_triton_rotary_emb_func":
149+
fn = lambda: rotary_embedding(new_q, new_k, cos, sin, k_cache, block_tables, kv_seq_lengths)
130150
else:
131151
raise ValueError("Undefined provider")
132152

@@ -135,5 +155,5 @@ def benchmark_rotary_emb(
135155

136156

137157
if __name__ == "__main__":
138-
test_rotary_emb(4, 64, 32, 64, torch.float32)
139-
# benchmark_rotary_emb.run(save_path=".",print_data=True)
158+
# test_rotary_emb(4, 64, 32, 64, torch.float32)
159+
benchmark_rotary_emb.run(save_path=".", print_data=True)

0 commit comments

Comments
 (0)