123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799800801802803804805806807808809810811812813814815816817818819820821822823824825826827828829830831832833834835836837838839840841842843844845846847848849850851852853854855856857858859860861862863864865866867868869870871872873874875876877878879880881882883884885886887888889890891892893894895896897898899900901902903904905906907908909910911912913914915916917918919920921922923924925926927928929930931932933934935936937938939940941942943944945946947948949950951952953954955956957958959960961962963964965966967968969970971972973974975976977978979980981982983984985986987988989990991992993994995996997998999100010011002100310041005100610071008100910101011101210131014101510161017101810191020102110221023102410251026102710281029103010311032103310341035103610371038103910401041104210431044104510461047104810491050105110521053105410551056105710581059106010611062106310641065106610671068106910701071107210731074107510761077107810791080108110821083108410851086108710881089109010911092109310941095109610971098109911001101110211031104110511061107110811091110111111121113111411151116111711181119112011211122112311241125112611271128112911301131113211331134113511361137113811391140114111421143114411451146114711481149115011511152115311541155115611571158115911601161116211631164116511661167116811691170117111721173117411751176117711781179118011811182118311841185118611871188118911901191119211931194119511961197119811991200120112021203120412051206120712081209121012111212121312141215121612171218121912201221122212231224122512261227122812291230123112321233123412351236123712381239124012411242124312441245124612471248124912501251125212531254125512561257125812591260126112621263126412651266126712681269127012711272127312741275127612771278127912801281128212831284128512861287128812891290129112921293129412951296129712981299130013011302130313041305130613071308130913101311131213131314131513161317131813191320132113221323132413251326132713281329133013311332133313341335133613371338133913401341134213431344134513461347134813491350135113521353135413551356135713581359136013611362136313641365136613671368136913701371137213731374137513761377137813791380138113821383138413851386138713881389139013911392139313941395139613971398139914001401140214031404140514061407140814091410141114121413141414151416141714181419142014211422142314241425142614271428142914301431143214331434143514361437143814391440144114421443144414451446144714481449145014511452145314541455145614571458145914601461146214631464146514661467146814691470147114721473147414751476147714781479148014811482148314841485148614871488148914901491149214931494149514961497149814991500150115021503150415051506150715081509151015111512151315141515151615171518151915201521152215231524152515261527152815291530153115321533153415351536153715381539154015411542154315441545154615471548154915501551155215531554155515561557155815591560156115621563156415651566156715681569157015711572157315741575157615771578157915801581158215831584158515861587158815891590159115921593159415951596159715981599160016011602160316041605160616071608160916101611161216131614161516161617161816191620 |
- import math
- import pytest
- import torch
- import torch.nn.functional as F
- from einops import rearrange, repeat
- from flash_attn import (
- flash_attn_func,
- flash_attn_kvpacked_func,
- flash_attn_qkvpacked_func,
- flash_attn_varlen_func,
- flash_attn_varlen_kvpacked_func,
- flash_attn_varlen_qkvpacked_func,
- flash_attn_with_kvcache,
- )
- from test_flash_attn import (
- attn_bias_from_alibi_slopes,
- convert_flash_attn_S_to_softmax,
- generate_qkv,
- generate_random_padding_mask,
- _generate_block_kvcache,
- attention_ref,
- attention_kvpacked_ref,
- attention_qkvpacked_ref,
- )
- from flash_attn.layers.rotary import apply_rotary_emb
- def is_bwd_hdim_supported(d):
- return d <= 256
- def ck_randval_to_dropout_mask(randval, p):
- # If p = 0.3, randval in 255 * (0.7, 1.0] will be dropout
- # randval in 255 * [0, 0.7] will be kept
- # If return dropout_mask >=0, value will be kept
- return math.floor(255.0 * (1 - p)) - randval.to(torch.float32)
- def pad_rearrange_dropout_mask_hts_to_bhss(S_dmask, cu_seqlens_q, seqlen_q_rounded, seqlen_k_rounded):
- """ pad + rearrange [nheads, total_q, max_seqlen_k] into [b, nheads, seqlen_q_rounded, seqlen_k_rounded]
- Arguments:
- S_dmask: (nheads, total_q, max_seqlen_k)
- cu_seqlens_q: (b + 1)
- Output:
- S_dmask: (b, nheads, seqlen_q_rounded, seqlen_k_rounded)
- """
- batch_size = cu_seqlens_q.numel() - 1
- seqlens_q = torch.roll(cu_seqlens_q, shifts = -1) - cu_seqlens_q
- seqlens_q = seqlens_q[0:batch_size].tolist()
- S_dmask = torch.split(S_dmask, seqlens_q, dim=1)
- # [(nheads, seqlen_q0, max_seqlen_k), (nheads, seqlen_q1, max_seqlen_k), ..., (nheads, seqlen_qb, max_seqlen_k)]
- masks = ()
- for mask in S_dmask:
- # (nheads, seqlen_qi, max_seqlen_k) -> (nheads, seqlen_q_rounded, seqlen_k_rounded)
- mask = F.pad(mask, (0, seqlen_k_rounded - mask.shape[2], 0, seqlen_q_rounded - mask.shape[1], 0, 0)).unsqueeze(1)
- masks = masks + (mask, )
- S_dmask = torch.cat(masks, dim=1)
- S_dmask = S_dmask.transpose(0, 1)
- return S_dmask
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("deterministic", [False, True])
- @pytest.mark.parametrize("alibi", [False, True])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize("seqlen", [97, 128, 200, 384, 768, 1024, 1025, 2048])
- @pytest.mark.parametrize("dropout_p", [0.0, 0.17])
- def test_flash_attn_qkvpacked(seqlen, d, dropout_p, causal, local, alibi, deterministic, dtype):
- if d > 256:
- pytest.skip()
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 4
- nheads = 9
- window_size = (-1, -1) if not local else torch.randint(0, seqlen, (2,))
- qkv = torch.randn(
- batch_size, seqlen, 3, nheads, d, device=device, dtype=dtype, requires_grad=True
- )
- if alibi:
- alibi_slopes = torch.rand(batch_size, nheads, device=device, dtype=torch.float32) * 0.3
- attn_bias = attn_bias_from_alibi_slopes(alibi_slopes, seqlen, seqlen, causal=causal)
- else:
- alibi_slopes, attn_bias = None, None
- out, lse, S_dmask = flash_attn_qkvpacked_func(
- qkv,
- dropout_p,
- causal=causal,
- window_size=window_size,
- alibi_slopes=alibi_slopes,
- deterministic=deterministic,
- return_attn_probs=True,
- )
- if dropout_p > 0.0:
- # TODO - move to c++ mha_varlen_fwd()
- S_dmask = ck_randval_to_dropout_mask(S_dmask, dropout_p)
- S_dmask_converted = convert_flash_attn_S_to_softmax(
- S_dmask,
- seqlen,
- seqlen,
- None,
- None,
- d,
- dropout_p > 0.0,
- causal=causal,
- window_size=window_size,
- )
- dropout_mask = S_dmask_converted >= 0
- # CK does not return P. Hence, we don't test the attn here.
- else:
- dropout_mask = None
- out_ref, attn_ref = attention_qkvpacked_ref(
- qkv, None, attn_bias, dropout_p, dropout_mask, causal=causal, window_size=window_size
- )
- out_pt, attn_pt = attention_qkvpacked_ref(
- qkv,
- None,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most twice the numerical error
- # of a Pytorch implementation.
- assert (out - out_ref).abs().max().item() <= 2 * (out_pt - out_ref).abs().max().item()
- g = torch.randn_like(out)
- if is_bwd_hdim_supported(d):
- (dqkv,) = torch.autograd.grad(out, qkv, g)
- (dqkv_ref,) = torch.autograd.grad(out_ref, qkv, g)
- (dqkv_pt,) = torch.autograd.grad(out_pt, qkv, g)
- print(f"dQ max diff: {(dqkv[:, :, 0] - dqkv_ref[:, :, 0]).abs().max().item()}")
- print(f"dK max diff: {(dqkv[:, :, 1] - dqkv_ref[:, :, 1]).abs().max().item()}")
- print(f"dV max diff: {(dqkv[:, :, 2] - dqkv_ref[:, :, 2]).abs().max().item()}")
- print(f"dQKV mean diff: {(dqkv - dqkv_ref).abs().mean().item()}")
- print(f"dQ Pytorch max diff: {(dqkv_pt[:, :, 0] - dqkv_ref[:, :, 0]).abs().max().item()}")
- print(f"dK Pytorch max diff: {(dqkv_pt[:, :, 1] - dqkv_ref[:, :, 1]).abs().max().item()}")
- print(f"dV Pytorch max diff: {(dqkv_pt[:, :, 2] - dqkv_ref[:, :, 2]).abs().max().item()}")
- print(f"dQKV Pytorch mean diff: {(dqkv_pt - dqkv_ref).abs().mean().item()}")
- # TODO - use 10 times to check, wait for ck to fix bwd precision issue
- assert (dqkv - dqkv_ref).abs().max().item() <= 10 * (dqkv_pt - dqkv_ref).abs().max().item()
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("deterministic", [False, True])
- @pytest.mark.parametrize("alibi", [False, True])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 59, 64, 80, 96, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize("seqlen", [97, 128, 200, 257, 384, 512, 768, 1025, 2048])
- @pytest.mark.parametrize("dropout_p", [0, 0.17])
- def test_flash_attn_varlen_qkvpacked(seqlen, d, dropout_p, causal, local, alibi, deterministic, dtype):
- if d > 256:
- pytest.skip()
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 5
- nheads = 6
- window_size = (-1, -1) if not local else torch.randint(0, seqlen, (2,))
- qkv = torch.randn(
- batch_size, seqlen, 3, nheads, d, device=device, dtype=dtype, requires_grad=True
- )
- key_padding_mask = generate_random_padding_mask(seqlen, batch_size, device, mode="random")
- # key_padding_mask = generate_random_padding_mask(seqlen, batch_size, device, mode='full')
- if alibi:
- alibi_slopes = torch.rand(batch_size, nheads, device=device, dtype=torch.float32) * 0.3
- attn_bias = attn_bias_from_alibi_slopes(
- alibi_slopes, seqlen, seqlen, key_padding_mask, key_padding_mask, causal=causal
- )
- else:
- alibi_slopes, attn_bias = None, None
- qkv_unpad, cu_seqlens, max_seqlen, qkv, output_pad_fn, dqkv_pad_fn = generate_qkv(
- *qkv.unbind(dim=2), key_padding_mask, key_padding_mask, qkvpacked=True
- )
- out_unpad, sm_lse, S_dmask = flash_attn_varlen_qkvpacked_func(
- qkv_unpad,
- cu_seqlens,
- max_seqlen,
- dropout_p,
- causal=causal,
- window_size=window_size,
- alibi_slopes=alibi_slopes,
- deterministic=deterministic,
- return_attn_probs=True,
- )
- out = output_pad_fn(out_unpad)
- if dropout_p > 0.0:
- # TODO - move to c++ mha_varlen_fwd()
- S_dmask = ck_randval_to_dropout_mask(S_dmask, dropout_p)
- S_dmask = pad_rearrange_dropout_mask_hts_to_bhss(S_dmask, cu_seqlens, seqlen, seqlen)
- S_dmask_converted = convert_flash_attn_S_to_softmax(
- S_dmask,
- seqlen,
- seqlen,
- key_padding_mask,
- key_padding_mask,
- d,
- dropout_p > 0.0,
- causal=causal,
- window_size=window_size,
- )
- dropout_mask = S_dmask_converted >= 0
- # CK does not return P. Hence, we don't test the attn here.
- else:
- dropout_mask = None
- out_ref, attn_ref = attention_qkvpacked_ref(
- qkv,
- key_padding_mask,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- )
- out_pt, attn_pt = attention_qkvpacked_ref(
- qkv,
- key_padding_mask,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most twice the numerical error
- # of a Pytorch implementation.
- assert (out - out_ref).abs().max().item() <= 2 * (out_pt - out_ref).abs().max().item()
- g = torch.randn_like(out)
- if is_bwd_hdim_supported(d):
- (dqkv_unpad,) = torch.autograd.grad(out, qkv_unpad, g)
- dqkv = dqkv_pad_fn(dqkv_unpad)
- (dqkv_ref,) = torch.autograd.grad(out_ref, qkv, g)
- (dqkv_pt,) = torch.autograd.grad(out_pt, qkv, g)
- print(f"dQ max diff: {(dqkv[:, :, 0] - dqkv_ref[:, :, 0]).abs().max().item()}")
- print(f"dK max diff: {(dqkv[:, :, 1] - dqkv_ref[:, :, 1]).abs().max().item()}")
- print(f"dV max diff: {(dqkv[:, :, 2] - dqkv_ref[:, :, 2]).abs().max().item()}")
- print(f"dQKV mean diff: {(dqkv - dqkv_ref).abs().mean().item()}")
- print(f"dQ Pytorch max diff: {(dqkv_pt[:, :, 0] - dqkv_ref[:, :, 0]).abs().max().item()}")
- print(f"dK Pytorch max diff: {(dqkv_pt[:, :, 1] - dqkv_ref[:, :, 1]).abs().max().item()}")
- print(f"dV Pytorch max diff: {(dqkv_pt[:, :, 2] - dqkv_ref[:, :, 2]).abs().max().item()}")
- print(f"dQKV Pytorch mean diff: {(dqkv_pt - dqkv_ref).abs().mean().item()}")
- # TODO - use 10 times to check, wait for ck to fix bwd precision issue
- assert (dqkv - dqkv_ref).abs().max().item() <= 10 * (dqkv_pt - dqkv_ref).abs().max().item()
- @pytest.mark.parametrize("kvpacked", [True, False])
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("mha_type", ["mha", "mqa", "gqa"])
- @pytest.mark.parametrize("deterministic", [False, True])
- @pytest.mark.parametrize("alibi", [False, True])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- (113, 203),
- (128, 217),
- (113, 211),
- (108, 256),
- (256, 512),
- (512, 256),
- (1024, 1024),
- (1023, 1024),
- (1024, 1023),
- (2048, 2048),
- ],
- )
- @pytest.mark.parametrize("dropout_p", [0.0, 0.17])
- def test_flash_attn_output(
- seqlen_q, seqlen_k, d, dropout_p, causal, local, alibi, deterministic, mha_type, dtype, kvpacked
- ):
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 4
- nheads = 9
- nheads_k = nheads if mha_type == "mha" else (1 if mha_type == "mqa" else 3)
- assert nheads % nheads_k == 0
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- if kvpacked:
- kv = torch.randn(
- batch_size, seqlen_k, 2, nheads_k, d, device=device, dtype=dtype, requires_grad=True
- )
- else:
- k = torch.randn(
- batch_size, seqlen_k, nheads_k, d, device=device, dtype=dtype, requires_grad=True
- )
- v = torch.randn(
- batch_size, seqlen_k, nheads_k, d, device=device, dtype=dtype, requires_grad=True
- )
- if alibi:
- alibi_slopes = torch.rand(batch_size, nheads, device=device, dtype=torch.float32) * 0.3
- attn_bias = attn_bias_from_alibi_slopes(alibi_slopes, seqlen_q, seqlen_k, causal=causal)
- else:
- alibi_slopes, attn_bias = None, None
- if kvpacked:
- out, lse, S_dmask = flash_attn_kvpacked_func(
- q,
- kv,
- dropout_p,
- causal=causal,
- window_size=window_size,
- alibi_slopes=alibi_slopes,
- deterministic=deterministic,
- return_attn_probs=True,
- )
- else:
- out, lse, S_dmask = flash_attn_func(
- q,
- k,
- v,
- dropout_p,
- causal=causal,
- window_size=window_size,
- alibi_slopes=alibi_slopes,
- deterministic=deterministic,
- return_attn_probs=True,
- )
- if dropout_p > 0.0:
- # TODO - move to c++ mha_varlen_fwd()
- S_dmask = ck_randval_to_dropout_mask(S_dmask, dropout_p)
- S_dmask_converted = convert_flash_attn_S_to_softmax(
- S_dmask,
- seqlen_q,
- seqlen_k,
- None,
- None,
- d,
- dropout_p > 0.0,
- causal=causal,
- window_size=window_size,
- )
- dropout_mask = S_dmask_converted >= 0
- if kvpacked:
- kv_rep = repeat(kv, "b s two h d -> b s two (h g) d", g=nheads // nheads_k)
- k_rep, v_rep = kv_rep.unbind(dim=2)
- else:
- k_rep = repeat(k, "b s h d -> b s (h g) d", g=nheads // nheads_k)
- v_rep = repeat(v, "b s h d -> b s (h g) d", g=nheads // nheads_k)
- # CK does not return P. Hence, we don't test the attn here.
- else:
- dropout_mask = None
- if kvpacked:
- out_ref, attn_ref = attention_kvpacked_ref(
- q,
- kv,
- None,
- None,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- )
- out_pt, attn_pt = attention_kvpacked_ref(
- q,
- kv,
- None,
- None,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- else:
- out_ref, attn_ref = attention_ref(
- q,
- k,
- v,
- None,
- None,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- )
- out_pt, attn_pt = attention_ref(
- q,
- k,
- v,
- None,
- None,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most twice the numerical error
- # of a Pytorch implementation.
- assert (out - out_ref).abs().max().item() <= 2 * (out_pt - out_ref).abs().max().item()
- g = torch.randn_like(out)
- if is_bwd_hdim_supported(d):
- if kvpacked:
- (
- dq,
- dkv,
- ) = torch.autograd.grad(out, (q, kv), g)
- dk, dv = dkv.unbind(2)
- (
- dq_ref,
- dkv_ref,
- ) = torch.autograd.grad(out_ref, (q, kv), g)
- dk_ref, dv_ref = dkv_ref.unbind(2)
- (
- dq_pt,
- dkv_pt,
- ) = torch.autograd.grad(out_pt, (q, kv), g)
- dk_pt, dv_pt = dkv_pt.unbind(2)
- else:
- (
- dq,
- dk,
- dv,
- ) = torch.autograd.grad(out, (q, k, v), g)
- (
- dq_ref,
- dk_ref,
- dv_ref,
- ) = torch.autograd.grad(out_ref, (q, k, v), g)
- (
- dq_pt,
- dk_pt,
- dv_pt,
- ) = torch.autograd.grad(out_pt, (q, k, v), g)
- print(f"dQ max diff: {(dq - dq_ref).abs().max().item()}")
- print(f"dK max diff: {(dk - dk_ref).abs().max().item()}")
- print(f"dV max diff: {(dv - dv_ref).abs().max().item()}")
- print(f"dQ mean diff: {(dq - dq_ref).abs().mean().item()}")
- print(f"dK mean diff: {(dk - dk_ref).abs().mean().item()}")
- print(f"dV mean diff: {(dv - dv_ref).abs().mean().item()}")
- print(f"dQ Pytorch max diff: {(dq_pt - dq_ref).abs().max().item()}")
- print(f"dK Pytorch max diff: {(dk_pt - dk_ref).abs().max().item()}")
- print(f"dV Pytorch max diff: {(dv_pt - dv_ref).abs().max().item()}")
- print(f"dQ Pytorch mean diff: {(dq_pt - dq_ref).abs().mean().item()}")
- print(f"dK Pytorch mean diff: {(dk_pt - dk_ref).abs().mean().item()}")
- print(f"dV Pytorch mean diff: {(dv_pt - dv_ref).abs().mean().item()}")
- # TODO - use 10 times to check, wait for ck to fix bwd precision issue
- assert (dq - dq_ref).abs().max().item() <= 10 * (dq_pt - dq_ref).abs().max().item()
- assert (dk - dk_ref).abs().max().item() <= 10 * (dk_pt - dk_ref).abs().max().item()
- assert (dv - dv_ref).abs().max().item() <= 10 * (dv_pt - dv_ref).abs().max().item()
- @pytest.mark.parametrize("kvpacked", [True, False])
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("mha_type", ["mha", "mqa", "gqa"])
- @pytest.mark.parametrize("deterministic", [False, True])
- @pytest.mark.parametrize("alibi", [False, True])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- (1, 147),
- (113, 203),
- (128, 217),
- (113, 211),
- (108, 256),
- (256, 512),
- (512, 256),
- (1024, 1024),
- (1023, 1024),
- (1024, 1023),
- (2048, 2048),
- ],
- )
- @pytest.mark.parametrize("dropout_p", [0.0, 0.17])
- def test_flash_attn_varlen_output(
- seqlen_q, seqlen_k, d, dropout_p, causal, local, alibi, deterministic, mha_type, dtype, kvpacked
- ):
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 4
- nheads = 9
- nheads_k = nheads if mha_type == "mha" else (1 if mha_type == "mqa" else 3)
- assert nheads % nheads_k == 0
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- if kvpacked:
- kv = torch.randn(
- batch_size, seqlen_k, 2, nheads_k, d, device=device, dtype=dtype, requires_grad=True
- )
- else:
- k = torch.randn(
- batch_size, seqlen_k, nheads_k, d, device=device, dtype=dtype, requires_grad=True
- )
- v = torch.randn(
- batch_size, seqlen_k, nheads_k, d, device=device, dtype=dtype, requires_grad=True
- )
- query_padding_mask = generate_random_padding_mask(seqlen_q, batch_size, device, mode="random")
- key_padding_mask = generate_random_padding_mask(seqlen_k, batch_size, device, mode="random")
- # key_padding_mask = generate_random_padding_mask(seqlen_k, batch_size, device, mode='full')
- if alibi:
- alibi_slopes = torch.rand(batch_size, nheads, device=device, dtype=torch.float32) * 0.3
- attn_bias = attn_bias_from_alibi_slopes(
- alibi_slopes, seqlen_q, seqlen_k, query_padding_mask, key_padding_mask, causal=causal
- )
- else:
- alibi_slopes, attn_bias = None, None
- if kvpacked:
- (
- q_unpad,
- kv_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- q,
- kv,
- output_pad_fn,
- dq_pad_fn,
- dkv_pad_fn,
- ) = generate_qkv(q, *kv.unbind(dim=2), query_padding_mask, key_padding_mask, kvpacked=True)
- out_unpad, sm_lse, S_dmask = flash_attn_varlen_kvpacked_func(
- q_unpad,
- kv_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- dropout_p,
- causal=causal,
- window_size=window_size,
- alibi_slopes=alibi_slopes,
- deterministic=deterministic,
- return_attn_probs=True,
- )
- else:
- (
- q_unpad,
- k_unpad,
- v_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- q,
- k,
- v,
- output_pad_fn,
- dq_pad_fn,
- dk_pad_fn,
- ) = generate_qkv(q, k, v, query_padding_mask, key_padding_mask, kvpacked=False)
- out_unpad, sm_lse, S_dmask = flash_attn_varlen_func(
- q_unpad,
- k_unpad,
- v_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- dropout_p,
- causal=causal,
- window_size=window_size,
- alibi_slopes=alibi_slopes,
- deterministic=deterministic,
- return_attn_probs=True,
- )
- out = output_pad_fn(out_unpad)
- if dropout_p > 0.0:
- # TODO - move to c++ mha_varlen_fwd()
- S_dmask = ck_randval_to_dropout_mask(S_dmask, dropout_p)
- S_dmask = pad_rearrange_dropout_mask_hts_to_bhss(S_dmask, cu_seqlens_q, seqlen_q, seqlen_k)
- S_dmask_converted = convert_flash_attn_S_to_softmax(
- S_dmask,
- seqlen_q,
- seqlen_k,
- query_padding_mask,
- key_padding_mask,
- d,
- dropout_p > 0.0,
- causal=causal,
- window_size=window_size,
- )
- dropout_mask = S_dmask_converted >= 0
- if kvpacked:
- kv_rep = repeat(kv, "b s two h d -> b s two (h g) d", g=nheads // nheads_k)
- k_rep, v_rep = kv_rep.unbind(dim=2)
- else:
- k_rep = repeat(k, "b s h d -> b s (h g) d", g=nheads // nheads_k)
- v_rep = repeat(v, "b s h d -> b s (h g) d", g=nheads // nheads_k)
- # CK does not return P. Hence, we don't test the attn here.
- else:
- dropout_mask = None
- if kvpacked:
- out_ref, attn_ref = attention_kvpacked_ref(
- q,
- kv,
- query_padding_mask,
- key_padding_mask,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- )
- out_pt, attn_pt = attention_kvpacked_ref(
- q,
- kv,
- query_padding_mask,
- key_padding_mask,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- else:
- out_ref, attn_ref = attention_ref(
- q,
- k,
- v,
- query_padding_mask,
- key_padding_mask,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- )
- out_pt, attn_pt = attention_ref(
- q,
- k,
- v,
- query_padding_mask,
- key_padding_mask,
- attn_bias,
- dropout_p,
- dropout_mask,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most 4 times the numerical error
- # of a Pytorch implementation.
- assert (out - out_ref).abs().max().item() <= 4 * (out_pt - out_ref).abs().max().item()
- g = torch.randn_like(out)
- if is_bwd_hdim_supported(d):
- if kvpacked:
- (
- dq_unpad,
- dkv_unpad,
- ) = torch.autograd.grad(out, (q_unpad, kv_unpad), g)
- dk, dv = dkv_pad_fn(dkv_unpad).unbind(2)
- (
- dq_ref,
- dkv_ref,
- ) = torch.autograd.grad(out_ref, (q, kv), g)
- dk_ref, dv_ref = dkv_ref.unbind(2)
- (
- dq_pt,
- dkv_pt,
- ) = torch.autograd.grad(out_pt, (q, kv), g)
- dk_pt, dv_pt = dkv_pt.unbind(2)
- else:
- (
- dq_unpad,
- dk_unpad,
- dv_unpad,
- ) = torch.autograd.grad(out, (q_unpad, k_unpad, v_unpad), g)
- dk = dk_pad_fn(dk_unpad)
- dv = dk_pad_fn(dv_unpad)
- (
- dq_ref,
- dk_ref,
- dv_ref,
- ) = torch.autograd.grad(out_ref, (q, k, v), g)
- (
- dq_pt,
- dk_pt,
- dv_pt,
- ) = torch.autograd.grad(out_pt, (q, k, v), g)
- dq = dq_pad_fn(dq_unpad)
- print(f"dQ max diff: {(dq - dq_ref).abs().max().item()}")
- print(f"dK max diff: {(dk - dk_ref).abs().max().item()}")
- print(f"dV max diff: {(dv - dv_ref).abs().max().item()}")
- print(f"dQ mean diff: {(dq - dq_ref).abs().mean().item()}")
- print(f"dK mean diff: {(dk - dk_ref).abs().mean().item()}")
- print(f"dV mean diff: {(dv - dv_ref).abs().mean().item()}")
- print(f"dQ Pytorch max diff: {(dq_pt - dq_ref).abs().max().item()}")
- print(f"dK Pytorch max diff: {(dk_pt - dk_ref).abs().max().item()}")
- print(f"dV Pytorch max diff: {(dv_pt - dv_ref).abs().max().item()}")
- print(f"dQ Pytorch mean diff: {(dq_pt - dq_ref).abs().mean().item()}")
- print(f"dK Pytorch mean diff: {(dk_pt - dk_ref).abs().mean().item()}")
- print(f"dV Pytorch mean diff: {(dv_pt - dv_ref).abs().mean().item()}")
- # TODO - use 10 times to check, wait for ck to fix bwd precision issue
- assert (dq - dq_ref).abs().max().item() <= 10 * (dq_pt - dq_ref).abs().max().item()
- assert (dk - dk_ref).abs().max().item() <= 10 * (dk_pt - dk_ref).abs().max().item()
- assert (dv - dv_ref).abs().max().item() <= 10 * (dv_pt - dv_ref).abs().max().item()
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize("swap_sq_sk", [False, True])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- # (1, 239),
- (3, 799),
- (127, 512),
- (127, 513),
- (113, 203),
- (128, 217),
- (113, 211),
- (108, 256),
- (256, 512),
- (1023, 1024),
- ],
- )
- def test_flash_attn_causal(seqlen_q, seqlen_k, swap_sq_sk, d, local, dtype):
- if max(seqlen_q, seqlen_k) >= 2048:
- pytest.skip()
- if swap_sq_sk:
- seqlen_q, seqlen_k = seqlen_k, seqlen_q
- device = "cuda"
- causal = True
- # set seed
- torch.random.manual_seed(0)
- batch_size = 8
- nheads = 9
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- k = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- v = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- out = flash_attn_func(q, k, v, 0.0, causal=causal, window_size=window_size)
- out_ref, attn_ref = attention_ref(
- q, k, v, None, None, None, 0.0, None, causal=causal, window_size=window_size
- )
- out_pt, attn_pt = attention_ref(
- q,
- k,
- v,
- None,
- None,
- None,
- 0.0,
- None,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most 4 times the numerical error
- # of a Pytorch implementation.
- assert (out - out_ref).abs().max().item() <= 4 * (out_pt - out_ref).abs().max().item() + 1e-5
- g = torch.randn_like(out)
- if is_bwd_hdim_supported(d):
- do_o = (g.float() * out.float()).sum(-1)
- (
- dq,
- dk,
- dv,
- ) = torch.autograd.grad(out, (q, k, v), g)
- (
- dq_ref,
- dk_ref,
- dv_ref,
- ) = torch.autograd.grad(out_ref, (q, k, v), g)
- (
- dq_pt,
- dk_pt,
- dv_pt,
- ) = torch.autograd.grad(out_pt, (q, k, v), g)
- print(f"dQ max diff: {(dq - dq_ref).abs().max().item()}")
- print(f"dK max diff: {(dk - dk_ref).abs().max().item()}")
- print(f"dV max diff: {(dv - dv_ref).abs().max().item()}")
- print(f"dQ mean diff: {(dq - dq_ref).abs().mean().item()}")
- print(f"dK mean diff: {(dk - dk_ref).abs().mean().item()}")
- print(f"dV mean diff: {(dv - dv_ref).abs().mean().item()}")
- print(f"dQ Pytorch max diff: {(dq_pt - dq_ref).abs().max().item()}")
- print(f"dK Pytorch max diff: {(dk_pt - dk_ref).abs().max().item()}")
- print(f"dV Pytorch max diff: {(dv_pt - dv_ref).abs().max().item()}")
- print(f"dQ Pytorch mean diff: {(dq_pt - dq_ref).abs().mean().item()}")
- print(f"dK Pytorch mean diff: {(dk_pt - dk_ref).abs().mean().item()}")
- print(f"dV Pytorch mean diff: {(dv_pt - dv_ref).abs().mean().item()}")
- # TODO - use 10 times to check, wait for ck to fix bwd precision issue
- assert (dq - dq_ref).abs().max().item() <= 10 * (dq_pt - dq_ref).abs().max().item() + 1e-4
- assert (dk - dk_ref).abs().max().item() <= 10 * (dk_pt - dk_ref).abs().max().item() + 1e-4
- assert (dv - dv_ref).abs().max().item() <= 10 * (dv_pt - dv_ref).abs().max().item() + 1e-4
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize("swap_sq_sk", [False, True])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- # (1, 239),
- (3, 799),
- (127, 512),
- (127, 513),
- (113, 203),
- (128, 217),
- (113, 211),
- (108, 256),
- (256, 512),
- (1023, 1024),
- ],
- )
- # TODO: Support paged_kv_block
- # @pytest.mark.parametrize("paged_kv_block_size", [None, 256, 512])
- @pytest.mark.parametrize("paged_kv_block_size", [None])
- def test_flash_attn_varlen_causal(
- seqlen_q, seqlen_k, swap_sq_sk, d, local, paged_kv_block_size, dtype
- ):
- if max(seqlen_q, seqlen_k) >= 2048:
- pytest.skip()
- if swap_sq_sk:
- seqlen_q, seqlen_k = seqlen_k, seqlen_q
- device = "cuda"
- causal = True
- # set seed
- torch.random.manual_seed(0)
- batch_size = 8
- nheads = 9
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- if paged_kv_block_size is None:
- k = torch.randn(
- batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True
- )
- v = torch.randn(
- batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True
- )
- block_table = None
- else:
- k, v, block_table, k_cache_paged, v_cache_paged, num_blocks = _generate_block_kvcache(
- seqlen_k, paged_kv_block_size, batch_size, nheads, d, device, dtype
- )
- query_padding_mask = generate_random_padding_mask(seqlen_q, batch_size, device, mode="random")
- key_padding_mask = generate_random_padding_mask(seqlen_k, batch_size, device, mode="random")
- (
- q_unpad,
- k_unpad,
- v_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- q,
- k,
- v,
- output_pad_fn,
- dq_pad_fn,
- dk_pad_fn,
- ) = generate_qkv(q, k, v, query_padding_mask, key_padding_mask, kvpacked=False)
- out_unpad = flash_attn_varlen_func(
- q_unpad,
- k_unpad if paged_kv_block_size is None else k_cache_paged,
- v_unpad if paged_kv_block_size is None else v_cache_paged,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- 0.0,
- causal=causal,
- window_size=window_size,
- block_table=block_table,
- )
- out = output_pad_fn(out_unpad)
- out_ref, attn_ref = attention_ref(
- q,
- k,
- v,
- query_padding_mask,
- key_padding_mask,
- None,
- 0.0,
- None,
- causal=causal,
- window_size=window_size,
- )
- out_pt, attn_pt = attention_ref(
- q,
- k,
- v,
- query_padding_mask,
- key_padding_mask,
- None,
- 0.0,
- None,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most twice the numerical error
- # of a Pytorch implementation.
- assert (out - out_ref).abs().max().item() <= 2 * (out_pt - out_ref).abs().max().item() + 1e-5
- g = torch.randn_like(out)
- if is_bwd_hdim_supported(d):
- do_o = (g.float() * out.float()).sum(-1)
- test_backward = block_table is None
- if test_backward:
- (
- dq_unpad,
- dk_unpad,
- dv_unpad,
- ) = torch.autograd.grad(out, (q_unpad, k_unpad, v_unpad), g)
- dq = dq_pad_fn(dq_unpad)
- dk = dk_pad_fn(dk_unpad)
- dv = dk_pad_fn(dv_unpad)
- (
- dq_ref,
- dk_ref,
- dv_ref,
- ) = torch.autograd.grad(out_ref, (q, k, v), g)
- (
- dq_pt,
- dk_pt,
- dv_pt,
- ) = torch.autograd.grad(out_pt, (q, k, v), g)
- print(f"dQ max diff: {(dq - dq_ref).abs().max().item()}")
- print(f"dK max diff: {(dk - dk_ref).abs().max().item()}")
- print(f"dV max diff: {(dv - dv_ref).abs().max().item()}")
- print(f"dQ mean diff: {(dq - dq_ref).abs().mean().item()}")
- print(f"dK mean diff: {(dk - dk_ref).abs().mean().item()}")
- print(f"dV mean diff: {(dv - dv_ref).abs().mean().item()}")
- print(f"dQ Pytorch max diff: {(dq_pt - dq_ref).abs().max().item()}")
- print(f"dK Pytorch max diff: {(dk_pt - dk_ref).abs().max().item()}")
- print(f"dV Pytorch max diff: {(dv_pt - dv_ref).abs().max().item()}")
- print(f"dQ Pytorch mean diff: {(dq_pt - dq_ref).abs().mean().item()}")
- print(f"dK Pytorch mean diff: {(dk_pt - dk_ref).abs().mean().item()}")
- print(f"dV Pytorch mean diff: {(dv_pt - dv_ref).abs().mean().item()}")
- if test_backward:
- # TODO - use 10 times to check, wait for ck to fix bwd precision issue
- assert (dq - dq_ref).abs().max().item() <= 10 * (dq_pt - dq_ref).abs().max().item() + 1e-5
- assert (dk - dk_ref).abs().max().item() <= 10 * (dk_pt - dk_ref).abs().max().item() + 1e-5
- assert (dv - dv_ref).abs().max().item() <= 10 * (dv_pt - dv_ref).abs().max().item() + 1e-5
- # TODO - support splitkv
- # def test_flash_attn_splitkv
- # TODO - Support has_leftpad
- @pytest.mark.parametrize("dtype", [torch.float16])
- @pytest.mark.parametrize("num_splits", [1, 0])
- @pytest.mark.parametrize("mha_type", ["mha", "mqa", "gqa"])
- @pytest.mark.parametrize("new_kv", [False, True])
- @pytest.mark.parametrize("alibi", [False, True])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("seqlen_new_eq_seqlen_q", [True, False])
- @pytest.mark.parametrize("rotary_interleaved", [False, True])
- @pytest.mark.parametrize("rotary_fraction", [0.0, 0.5, 1.0])
- @pytest.mark.parametrize("paged_kv_block_size", [None, 256])
- @pytest.mark.parametrize("has_leftpad", [False])
- @pytest.mark.parametrize("has_batch_idx", [False, True])
- @pytest.mark.parametrize("d", [32, 59, 64, 80, 128, 256])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- (1, 128),
- (1, 339),
- (3, 1024),
- (64, 800),
- (64, 256),
- (3, 799),
- (64, 2048),
- (16, 20000),
- (1, 128 * 1024),
- (16, 128 * 1024),
- (128, 128),
- ],
- )
- def test_flash_attn_kvcache(
- seqlen_q,
- seqlen_k,
- d,
- has_batch_idx,
- has_leftpad,
- paged_kv_block_size,
- rotary_fraction,
- rotary_interleaved,
- seqlen_new_eq_seqlen_q,
- causal,
- local,
- alibi,
- new_kv,
- mha_type,
- num_splits,
- dtype,
- ):
- if seqlen_q > seqlen_k and new_kv:
- pytest.skip()
- if not new_kv and rotary_fraction > 0.0:
- pytest.skip()
- if has_batch_idx and paged_kv_block_size is not None:
- pytest.skip()
- if has_leftpad and paged_kv_block_size is not None:
- pytest.skip()
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 1
- batch_size_cache = batch_size if not has_batch_idx else batch_size * 2
- nheads = 6
- # rotary_dim must be a multiple of 16, and must be <= d
- rotary_dim = math.floor(int(rotary_fraction * d) / 16) * 16
- nheads_k = nheads if mha_type == "mha" else (1 if mha_type == "mqa" else 3)
- assert nheads % nheads_k == 0
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype)
- seqlen_new = seqlen_q if seqlen_new_eq_seqlen_q else torch.randint(1, seqlen_q + 1, (1,)).item()
- if new_kv:
- k = torch.randn(batch_size, seqlen_new, nheads_k, d, device=device, dtype=dtype)
- v = torch.randn(batch_size, seqlen_new, nheads_k, d, device=device, dtype=dtype)
- else:
- k, v = None, None
- if paged_kv_block_size is None:
- k_cache = torch.randn(batch_size_cache, seqlen_k, nheads_k, d, device=device, dtype=dtype)
- v_cache = torch.randn(batch_size_cache, seqlen_k, nheads_k, d, device=device, dtype=dtype)
- block_table = None
- else:
- (
- k_cache,
- v_cache,
- block_table,
- k_cache_paged,
- v_cache_paged,
- num_blocks,
- ) = _generate_block_kvcache(
- seqlen_k, paged_kv_block_size, batch_size, nheads_k, d, device, dtype
- )
- cache_seqlens = torch.randint(
- 0 if new_kv else 1,
- # If we don't use seqlen_q in the case of causal and rotary, cos/sin won't be long enough
- (
- (seqlen_k - (seqlen_q if (causal or local) and rotary_dim > 1 else seqlen_new) + 1)
- if new_kv
- else (seqlen_k + 1)
- ),
- (batch_size,),
- dtype=torch.int32,
- device=device,
- )
- if has_leftpad:
- cache_leftpad = torch.cat([torch.randint(0, cache_seqlens[i].item(), (1,), dtype=torch.int32, device=device)
- if cache_seqlens[i].item() > 0 else torch.zeros(1, dtype=torch.int32, device=device)
- for i in range(batch_size)])
- else:
- cache_leftpad = None
- arange = rearrange(torch.arange(seqlen_k, device=device), "s -> 1 s")
- cache_seqlens_expanded = rearrange(cache_seqlens, "b -> b 1")
- key_padding_mask = arange < cache_seqlens_expanded + (seqlen_new if new_kv else 0)
- if has_leftpad:
- key_padding_mask = torch.logical_and(
- key_padding_mask, arange >= cache_leftpad.unsqueeze(-1).expand(-1, seqlen_k)
- )
- if has_batch_idx:
- cache_batch_idx = torch.randperm(batch_size_cache, dtype=torch.int32, device=device)[
- :batch_size
- ]
- else:
- cache_batch_idx = None
- if alibi:
- alibi_slopes = torch.rand(batch_size, nheads, device=device, dtype=torch.float32) * 0.3
- attn_bias = attn_bias_from_alibi_slopes(
- alibi_slopes, seqlen_q, seqlen_k, None, key_padding_mask, causal=causal, key_leftpad=cache_leftpad
- )
- else:
- alibi_slopes, attn_bias = None, None
- # cache_seqlens = torch.tensor([64], dtype=torch.int32, device=device)
- if rotary_dim > 0:
- angle = (
- torch.rand(
- seqlen_k if paged_kv_block_size is None else num_blocks * paged_kv_block_size,
- rotary_dim // 2,
- device=device,
- )
- * 2
- * math.pi
- )
- cos = torch.cos(angle).to(dtype=dtype)
- sin = torch.sin(angle).to(dtype=dtype)
- if causal or local:
- q_ro = apply_rotary_emb(
- q, cos, sin, seqlen_offsets=cache_seqlens, interleaved=rotary_interleaved
- )
- else:
- q_ro = rearrange(
- apply_rotary_emb(
- rearrange(q, "b s h d -> b 1 (s h) d"),
- cos,
- sin,
- seqlen_offsets=cache_seqlens,
- interleaved=rotary_interleaved,
- ),
- "b 1 (s h) d -> b s h d",
- s=seqlen_q,
- )
- # q_ro = q
- k_ro = apply_rotary_emb(
- k, cos, sin, seqlen_offsets=cache_seqlens, interleaved=rotary_interleaved
- )
- else:
- cos, sin = None, None
- q_ro, k_ro = q, k
- # k_cache[:, 64:] = -1
- k_cache_ref = (
- k_cache if not has_batch_idx else k_cache[cache_batch_idx.to(dtype=torch.long)]
- ).clone()
- v_cache_ref = (
- v_cache if not has_batch_idx else v_cache[cache_batch_idx.to(dtype=torch.long)]
- ).clone()
- if new_kv:
- update_mask = torch.logical_and(
- cache_seqlens_expanded <= arange, arange < cache_seqlens_expanded + seqlen_new
- )
- k_cache_ref[update_mask] = rearrange(k_ro, "b s ... -> (b s) ...")
- v_cache_ref[update_mask] = rearrange(v, "b s ... -> (b s) ...")
- k_cache_rep = repeat(k_cache_ref, "b s h d -> b s (h g) d", g=nheads // nheads_k)
- v_cache_rep = repeat(v_cache_ref, "b s h d -> b s (h g) d", g=nheads // nheads_k)
- out = flash_attn_with_kvcache(
- q,
- k_cache if paged_kv_block_size is None else k_cache_paged,
- v_cache if paged_kv_block_size is None else v_cache_paged,
- k,
- v,
- rotary_cos=cos,
- rotary_sin=sin,
- cache_seqlens=cache_seqlens,
- cache_batch_idx=cache_batch_idx,
- cache_leftpad=cache_leftpad,
- block_table=block_table,
- causal=causal,
- window_size=window_size,
- rotary_interleaved=rotary_interleaved,
- alibi_slopes=alibi_slopes,
- num_splits=num_splits,
- )
- # out = flash_attn_with_kvcache(
- # q, k_cache, v_cache, cache_seqlens=cache_seqlens, causal=causal, window_size=window_size
- # )
- # out = flash_attn_with_kvcache(q, k_cache, v_cache, causal=causal, window_size=window_size)
- # qk = torch.einsum("bqhd,bkhd->bhqk", q, k_cache_ref)
- # m = qk.amax(-1, keepdim=True)
- # s_tmp = torch.exp((qk - m) / math.sqrt(d))
- # o1 = torch.einsum('bhst,bthd->bshd', s_tmp, v_cache_ref)
- # lse_ref = torch.logsumexp(qk / math.sqrt(d), -1)
- # probs = torch.softmax(qk, dim=-1)
- out_ref, _ = attention_ref(
- q_ro,
- k_cache_rep,
- v_cache_rep,
- None,
- key_padding_mask,
- attn_bias,
- 0.0,
- None,
- causal=causal,
- window_size=window_size,
- key_leftpad=cache_leftpad,
- )
- out_pt, _ = attention_ref(
- q_ro,
- k_cache_rep,
- v_cache_rep,
- None,
- key_padding_mask,
- attn_bias,
- 0.0,
- None,
- causal=causal,
- window_size=window_size,
- upcast=False,
- reorder_ops=True,
- key_leftpad=cache_leftpad,
- )
- print(f"Output max diff: {(out - out_ref).abs().max().item()}")
- print(f"Output mean diff: {(out - out_ref).abs().mean().item()}")
- print(f"Pytorch max diff: {(out_pt - out_ref).abs().max().item()}")
- print(f"Pytorch mean diff: {(out_pt - out_ref).abs().mean().item()}")
- # Check that FlashAttention's numerical error is at most twice the numerical error
- # of a Pytorch implementation.
- if new_kv:
- if paged_kv_block_size is None:
- k_cache_select = (
- k_cache if not has_batch_idx else k_cache[cache_batch_idx.to(dtype=torch.long)]
- )
- v_cache_select = (
- v_cache if not has_batch_idx else v_cache[cache_batch_idx.to(dtype=torch.long)]
- )
- else:
- k_cache_select = rearrange(
- k_cache_paged[block_table.to(dtype=torch.long).flatten()],
- "(b nblocks) block_size ... -> b (nblocks block_size) ...",
- b=batch_size,
- )[:, :seqlen_k]
- v_cache_select = rearrange(
- v_cache_paged[block_table.to(dtype=torch.long).flatten()],
- "(b nblocks) block_size ... -> b (nblocks block_size) ...",
- b=batch_size,
- )[:, :seqlen_k]
- assert torch.allclose(k_cache_select, k_cache_ref, rtol=1e-3, atol=1e-3)
- assert torch.equal(v_cache_select, v_cache_ref)
- # mult = 3 if f16, bf16 need 4
- mult = 4 if not alibi else 5
- assert (out - out_ref).abs().max().item() <= mult * (out_pt - out_ref).abs().max().item() + 1e-5
- @pytest.mark.parametrize("dtype", [torch.float16])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- (1, 239),
- (239, 1),
- (3, 799),
- (799, 3),
- (1024, 128),
- (97, 97),
- (128, 128),
- (200, 200),
- (256, 256),
- (257, 257),
- (384, 384),
- (512, 512),
- (768, 768),
- # (1024, 1024),
- ],
- )
- @pytest.mark.parametrize("dropout_p", [0.0, 0.17])
- def test_flash_attn_race_condition(seqlen_q, seqlen_k, d, dropout_p, causal, dtype):
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 60 # Sometimes we need large batch size for the race conditions to trigger
- nheads = 4
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- k = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- v = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- torch.random.manual_seed(42)
- out0, lse0, _ = flash_attn_func(q, k, v, dropout_p, causal=causal, return_attn_probs=True)
- g = torch.randn_like(out0)
- if dropout_p == 0 and is_bwd_hdim_supported(d):
- (
- dq0,
- dk0,
- dv0,
- ) = torch.autograd.grad(out0, (q, k, v), g)
- # Numerical error if we just do any arithmetic on dq
- dq_atol = 2 * ((dq0 + 0.3 - 0.3) - dq0).abs().max().item()
- for i in range(250):
- torch.random.manual_seed(42)
- out, lse, _ = flash_attn_func(q, k, v, dropout_p, causal=causal, return_attn_probs=True)
- assert torch.equal(out, out0)
- assert torch.equal(lse, lse0)
- if dropout_p == 0:
- (
- dq,
- dk,
- dv,
- ) = torch.autograd.grad(out, (q, k, v), g)
- dq_equal = torch.allclose(dq, dq0, atol=dq_atol)
- if not dq_equal:
- print(f"Iter {i}, {dq_atol = }, dQ max diff: {(dq - dq0).abs().max().item()}")
- assert torch.equal(dv, dv0)
- assert torch.equal(dk, dk0)
- assert dq_equal
- @pytest.mark.parametrize("dtype", [torch.float16])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [16, 32, 64])
- @pytest.mark.parametrize("seqlen", [1, 2, 5, 17, 128])
- def test_flash_attn_bwd_overflow(seqlen, d, causal, dtype):
- """We previously had a bug where not masking elements beyond seqlen_k caused NaN in dQ,
- in the case where seqlen % 128 != 0.
- """
- # TODO - 1 or 2 might fail, need to check
- if seqlen == 1 or seqlen == 2:
- pytest.skip()
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 2
- nheads = 5
- q = torch.randn([batch_size, seqlen, nheads, d], dtype=dtype, device="cuda") * 5
- k, v = [
- torch.randn([batch_size, seqlen, nheads, d], dtype=dtype, device="cuda") * 3
- for _ in range(2)
- ]
- q.requires_grad_(True)
- k.requires_grad_(True)
- v.requires_grad_(True)
- out = flash_attn_func(q, k, v, causal=causal)
- g = torch.randn_like(out)
- out.backward(g)
- q_pt = q.detach().clone().requires_grad_(True)
- k_pt = k.detach().clone().requires_grad_(True)
- v_pt = v.detach().clone().requires_grad_(True)
- out_pt, _ = attention_ref(q_pt, k_pt, v_pt, causal=causal, upcast=False, reorder_ops=True)
- out_pt.backward(g)
- q_ref = q.detach().clone().requires_grad_(True)
- k_ref = k.detach().clone().requires_grad_(True)
- v_ref = v.detach().clone().requires_grad_(True)
- out_ref, attn_ref = attention_ref(q_ref, k_ref, v_ref, causal=causal)
- out_ref.backward(g)
- print(f"dQ max diff: {(q.grad - q_ref.grad).abs().max().item()}")
- print(f"dK max diff: {(k.grad - k_ref.grad).abs().max().item()}")
- print(f"dV max diff: {(v.grad - v_ref.grad).abs().max().item()}")
- print(f"dQ Pytorch max diff: {(q_pt.grad - q_ref.grad).abs().max().item()}")
- print(f"dK Pytorch max diff: {(k_pt.grad - k_ref.grad).abs().max().item()}")
- print(f"dV Pytorch max diff: {(v_pt.grad - v_ref.grad).abs().max().item()}")
- assert (out - out_ref).abs().max().item() <= 2 * (out_pt - out_ref).abs().max().item()
- assert (q.grad - q_ref.grad).abs().max().item() <= 5 * (
- q_pt.grad - q_ref.grad
- ).abs().max().item() + 1e-3
- assert (k.grad - k_ref.grad).abs().max().item() <= 5 * (
- k_pt.grad - k_ref.grad
- ).abs().max().item() + 1e-3
- assert (v.grad - v_ref.grad).abs().max().item() <= 5 * (
- v_pt.grad - v_ref.grad
- ).abs().max().item() + 1e-3
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [64, 128])
- @pytest.mark.parametrize("seqlen", [97, 128, 200, 256])
- def test_flash_attn_bwd_transpose(seqlen, d, causal, dtype):
- """We previously had a bug where we were using the wrong strides of dout, which shows up
- when dout is not contiguous.
- """
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 5
- nheads = 2
- q, k, v = [
- torch.randn([batch_size, seqlen, nheads, d], dtype=dtype, device="cuda", requires_grad=True)
- for _ in range(3)
- ]
- out = rearrange(flash_attn_func(q, k, v, causal=causal), "b s ... -> s b ...")
- # So g is not contiguous
- g = torch.randn(seqlen, 2 * batch_size, nheads, d, dtype=dtype, device="cuda")[:, ::2]
- out.backward(g)
- q_pt = q.detach().clone().requires_grad_(True)
- k_pt = k.detach().clone().requires_grad_(True)
- v_pt = v.detach().clone().requires_grad_(True)
- out_pt, attn_pt = attention_ref(q_pt, k_pt, v_pt, causal=causal, upcast=False, reorder_ops=True)
- out_pt = rearrange(out_pt, "b s ... -> s b ...")
- out_pt.backward(g)
- q_ref = q.detach().clone().requires_grad_(True)
- k_ref = k.detach().clone().requires_grad_(True)
- v_ref = v.detach().clone().requires_grad_(True)
- out_ref, attn_ref = attention_ref(q_ref, k_ref, v_ref, causal=causal)
- out_ref = rearrange(out_ref, "b s ... -> s b ...")
- out_ref.backward(g)
- print(f"dQ max diff: {(q.grad - q_ref.grad).abs().max().item()}")
- print(f"dK max diff: {(k.grad - k_ref.grad).abs().max().item()}")
- print(f"dV max diff: {(v.grad - v_ref.grad).abs().max().item()}")
- print(f"dQ Pytorch max diff: {(q_pt.grad - q_ref.grad).abs().max().item()}")
- print(f"dK Pytorch max diff: {(k_pt.grad - k_ref.grad).abs().max().item()}")
- print(f"dV Pytorch max diff: {(v_pt.grad - v_ref.grad).abs().max().item()}")
- assert (out - out_ref).abs().max().item() <= 2 * (out_pt - out_ref).abs().max().item()
- assert (q.grad - q_ref.grad).abs().max().item() <= 2 * (
- q_pt.grad - q_ref.grad
- ).abs().max().item()
- assert (k.grad - k_ref.grad).abs().max().item() <= 2 * (
- k_pt.grad - k_ref.grad
- ).abs().max().item()
- assert (v.grad - v_ref.grad).abs().max().item() <= 2 * (
- v_pt.grad - v_ref.grad
- ).abs().max().item()
- @pytest.mark.parametrize("dtype", [torch.float16])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [16, 32, 64])
- def test_flash_attn_bwd_varlen_overflow(d, causal, dtype):
- """We previously had a bug where not masking elements beyond seqlen_k caused NaN in dQ,
- in the case where seqlen % 128 != 0 or varlen.
- """
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- nheads = 5
- q_cuseqlen = torch.tensor([0, 76, 110, 256], device=device, dtype=torch.int32)
- k_cuseqlen = torch.tensor([0, 1, 2, 3], device=device, dtype=torch.int32)
- Mq = 256
- Mk = 3
- q = torch.randn([Mq, nheads, d], dtype=dtype, device=device) * 3
- k, v = [torch.randn([Mk, nheads, d], dtype=dtype, device=device) * 3 for _ in range(2)]
- q.requires_grad_(True)
- k.requires_grad_(True)
- v.requires_grad_(True)
- out = flash_attn_varlen_func(q, k, v, q_cuseqlen, k_cuseqlen, Mq, Mk, causal=causal)
- g = torch.randn_like(out)
- out.backward(g)
- assert not q.grad.isnan().any()
- assert not k.grad.isnan().any()
- assert not v.grad.isnan().any()
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize("swap_sq_sk", [False, True])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- (1, 239),
- (3, 799),
- (127, 512),
- (127, 513),
- (113, 203),
- (128, 217),
- (113, 211),
- (108, 256),
- (256, 512),
- (1023, 1024),
- ],
- )
- def test_flash_attn_deterministic(seqlen_q, seqlen_k, swap_sq_sk, d, causal, local, dtype):
- if (
- max(seqlen_q, seqlen_k) >= 2048
- and torch.cuda.get_device_properties("cuda").total_memory <= 16 * 2**30
- ):
- pytest.skip() # Reference implementation OOM
- if swap_sq_sk:
- seqlen_q, seqlen_k = seqlen_k, seqlen_q
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 4
- nheads = 9
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- k = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- v = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- out = flash_attn_func(q, k, v, 0.0, causal=causal, window_size=window_size, deterministic=True)
- g = torch.randn_like(out)
- dq0, dk0, dv0 = torch.autograd.grad(out, (q, k, v), g, retain_graph=True)
- for _ in range(50):
- dq, dk, dv = torch.autograd.grad(out, (q, k, v), g, retain_graph=True)
- assert torch.equal(dv, dv0)
- assert torch.equal(dk, dk0)
- assert torch.equal(dq, dq0)
- @pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
- @pytest.mark.parametrize("local", [False, True])
- @pytest.mark.parametrize("causal", [False, True])
- @pytest.mark.parametrize("d", [32, 40, 59, 64, 80, 96, 111, 128, 160, 192, 224, 256])
- @pytest.mark.parametrize("swap_sq_sk", [False, True])
- @pytest.mark.parametrize(
- "seqlen_q,seqlen_k",
- [
- (1, 239),
- (3, 799),
- (127, 512),
- (127, 513),
- (113, 203),
- (128, 217),
- (113, 211),
- (108, 256),
- (256, 512),
- (1023, 1024),
- ],
- )
- def test_flash_attn_varlen_deterministic(seqlen_q, seqlen_k, swap_sq_sk, d, causal, local, dtype):
- if (
- max(seqlen_q, seqlen_k) >= 2048
- and torch.cuda.get_device_properties("cuda").total_memory <= 16 * 2**30
- ):
- pytest.skip() # Reference implementation OOM
- if swap_sq_sk:
- seqlen_q, seqlen_k = seqlen_k, seqlen_q
- device = "cuda"
- # set seed
- torch.random.manual_seed(0)
- batch_size = 2
- nheads = 9
- window_size = (-1, -1) if not local else torch.randint(0, seqlen_k, (2,))
- q = torch.randn(batch_size, seqlen_q, nheads, d, device=device, dtype=dtype, requires_grad=True)
- k = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- v = torch.randn(batch_size, seqlen_k, nheads, d, device=device, dtype=dtype, requires_grad=True)
- query_padding_mask = generate_random_padding_mask(seqlen_q, batch_size, device, mode="random")
- key_padding_mask = generate_random_padding_mask(seqlen_k, batch_size, device, mode="random")
- (
- q_unpad,
- k_unpad,
- v_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- q,
- k,
- v,
- output_pad_fn,
- dq_pad_fn,
- dk_pad_fn,
- ) = generate_qkv(q, k, v, query_padding_mask, key_padding_mask, kvpacked=False)
- out = flash_attn_varlen_func(
- q_unpad,
- k_unpad,
- v_unpad,
- cu_seqlens_q,
- cu_seqlens_k,
- max_seqlen_q,
- max_seqlen_k,
- 0.0,
- causal=causal,
- window_size=window_size,
- deterministic=True,
- )
- g = torch.randn_like(out)
- dq0, dk0, dv0 = torch.autograd.grad(out, (q_unpad, k_unpad, v_unpad), g, retain_graph=True)
- for _ in range(50):
- dq, dk, dv = torch.autograd.grad(out, (q_unpad, k_unpad, v_unpad), g, retain_graph=True)
- assert torch.equal(dv, dv0)
- assert torch.equal(dk, dk0)
- assert torch.equal(dq, dq0)
|