sdpa.py 34 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799800801802803804805806807808809810811812813814815816817818819820821822823824825826827828829830831832833834835836837838839840841842843844845846847848849850851852853854855856857858859860861862863864865866867868869870871872873874875876877878879880881882883884885886887888889890891892893894895896897898899900901902903904905906907908909910911912913914915916917918919920921922923924925926927928929930931932933934
  1. # mypy: allow-untyped-defs
  2. import logging
  3. from typing import Optional
  4. import torch
  5. import torch.nn
  6. import torch.nn.functional as F
  7. from torch.backends.cuda import (
  8. can_use_cudnn_attention,
  9. can_use_efficient_attention,
  10. can_use_flash_attention,
  11. cudnn_sdp_enabled,
  12. flash_sdp_enabled,
  13. math_sdp_enabled,
  14. mem_efficient_sdp_enabled,
  15. SDPAParams,
  16. )
  17. from torch.nn.attention import SDPBackend
  18. from .nested_tensor import NestedTensor
  19. log = logging.getLogger(__name__)
  20. def _validate_sdpa_input(
  21. query: torch.Tensor,
  22. key: torch.Tensor,
  23. value: torch.Tensor,
  24. attn_mask: Optional[torch.Tensor] = None,
  25. dropout_p=0.0,
  26. is_causal=False,
  27. scale=None,
  28. ):
  29. if (
  30. not isinstance(query, NestedTensor)
  31. or not isinstance(key, NestedTensor)
  32. or not isinstance(value, NestedTensor)
  33. ):
  34. raise ValueError(
  35. f"Expected query, key, and value to be nested tensors, "
  36. f"but got query.is_nested: {query.is_nested}, key.is_nested: {key.is_nested}, "
  37. f"and value.is_nested: {value.is_nested} instead."
  38. )
  39. if query.dtype != key.dtype or query.dtype != value.dtype:
  40. raise ValueError(
  41. f"Expected query, key, and value to have the same dtype, "
  42. f"but got query.dtype: {query.dtype}, key.dtype: {key.dtype}, "
  43. f"and value.dtype: {value.dtype} instead."
  44. )
  45. if query.device != key.device or query.device != value.device:
  46. raise ValueError(
  47. f"Expected query, key, and value to have the same device type, "
  48. f"but got query.device: {query.device}, key.device: {key.device}, "
  49. f"and value.device: {value.device} instead."
  50. )
  51. if query.dim() < 3 or key.dim() < 3 or value.dim() < 3:
  52. raise ValueError(
  53. f"Expected query, key, and value to all be at least 3 dimensional, but got query.dim: "
  54. f"{query.dim()}, key.dim: {key.dim()} and value.dim: {value.dim()} instead."
  55. )
  56. if query._ragged_idx != key._ragged_idx or query._ragged_idx != value._ragged_idx:
  57. raise ValueError(
  58. f"Expected query, key, and value to all be ragged on the same dimension, but got ragged "
  59. f"dims {query._ragged_idx}, {key._ragged_idx}, and {value._ragged_idx}, respectively."
  60. )
  61. if attn_mask is not None:
  62. # TODO: Figure out whether masks are actually supported for this layout or not
  63. raise ValueError("Masks are not yet supported!")
  64. if attn_mask.dtype != torch.bool and attn_mask.dtype != query.dtype:
  65. raise ValueError(
  66. f"Expected attn_mask dtype to be bool or to match query dtype, but got attn_mask.dtype: "
  67. f"{attn_mask.dtype}, and query.dtype: {query.dtype} instead."
  68. )
  69. def _check_batch_size_nested(params: SDPAParams, debug=False) -> bool:
  70. # This is expected to be called after check_tensor_shapes ensuring that the
  71. # size() calls won't error since the inputs are all 4 dimensional
  72. q_batch_size = params.query.size(0)
  73. k_batch_size = params.key.size(0)
  74. v_batch_size = params.value.size(0)
  75. # num_heads logic for nested input is checked in
  76. # check_for_seq_len_0_nested_tensor as there is handling there to make sure
  77. # num_heads is not ragged
  78. return q_batch_size == k_batch_size and q_batch_size == v_batch_size
  79. def _check_head_dim_size_flash_nested(params: SDPAParams, debug=False) -> bool:
  80. max_size = 256
  81. query_size_last = params.query.size(-1)
  82. key_size_last = params.key.size(-1)
  83. value_size_last = params.value.size(-1)
  84. same_head_dim_size = (
  85. query_size_last == key_size_last and query_size_last == value_size_last
  86. )
  87. if not (
  88. same_head_dim_size
  89. and (query_size_last % 8 == 0)
  90. and (query_size_last <= max_size)
  91. ):
  92. if debug:
  93. log.warning(
  94. "For NestedTensor inputs, Flash attention requires q,k,v to have the same "
  95. "last dimension and to be a multiple of 8 and less than or equal to 256. "
  96. "Got Query.size(-1): %d, Key.size(-1): %d, Value.size(-1): %d instead.",
  97. query_size_last,
  98. key_size_last,
  99. value_size_last,
  100. )
  101. return False
  102. return True
  103. def _check_head_dim_size_cudnn_nested(params: SDPAParams, debug=False) -> bool:
  104. max_size = 128
  105. query_size_last = params.query.size(-1)
  106. key_size_last = params.key.size(-1)
  107. value_size_last = params.value.size(-1)
  108. same_head_dim_size = (
  109. query_size_last == key_size_last and query_size_last == value_size_last
  110. )
  111. if not (
  112. same_head_dim_size
  113. and (query_size_last % 8 == 0)
  114. and (query_size_last <= max_size)
  115. ):
  116. if debug:
  117. log.warning(
  118. "For NestedTensor inputs, cuDNN attention requires q,k,v to have the same "
  119. "last dimension and to be a multiple of 8 and less than or equal to 128. "
  120. "Got Query.size(-1): %d, Key.size(-1): %d, Value.size(-1): %d instead.",
  121. query_size_last,
  122. key_size_last,
  123. value_size_last,
  124. )
  125. return False
  126. return True
  127. def _check_for_seq_len_0_and_consistent_head_dim_nested_helper(
  128. param: torch.Tensor, param_name: str, debug=False
  129. ) -> bool:
  130. assert isinstance(param, NestedTensor), "param should be a jagged NT"
  131. if param._ragged_idx == 1:
  132. # num_head_dims is ragged
  133. if debug:
  134. log.warning(
  135. "Fused kernels do not support ragged num_head_dims, %s has a ragged num_heads.",
  136. param_name,
  137. )
  138. return False
  139. # This is being called inside sdp with shape [batch, heads, {seq_len}, dim]
  140. if param._get_min_seqlen() == 0:
  141. if debug:
  142. log.warning(
  143. "Fused kernels do not support seq_len == 0, %s has a seq len of 0.",
  144. param_name,
  145. )
  146. return False
  147. return True
  148. def _try_broadcast_param_size(q_size, k_size, v_size, param_name, debug=False) -> bool:
  149. max_size = max(q_size, k_size, v_size)
  150. if (
  151. (q_size != max_size and q_size != 1)
  152. or (k_size != max_size and k_size != 1)
  153. or (v_size != max_size and v_size != 1)
  154. ):
  155. if debug:
  156. log.warning(
  157. "Both fused kernels require query, key and value to have broadcastable %s, "
  158. "got Query %s %d, Key %s %d, Value %s %d instead.",
  159. param_name,
  160. param_name,
  161. q_size,
  162. param_name,
  163. k_size,
  164. param_name,
  165. v_size,
  166. )
  167. return False
  168. return True
  169. def _check_for_seq_len_0_nested(params: SDPAParams, debug=False) -> bool:
  170. # When this function is called we are assured that the nt is dim==4
  171. q_is_safe = (
  172. _check_for_seq_len_0_and_consistent_head_dim_nested_helper(
  173. params.query, "query", debug
  174. )
  175. if params.query.is_nested
  176. else True
  177. )
  178. # short circuit if any is unsafe
  179. if not q_is_safe:
  180. return False
  181. k_is_safe = (
  182. _check_for_seq_len_0_and_consistent_head_dim_nested_helper(
  183. params.key, "key", debug
  184. )
  185. if params.key.is_nested
  186. else True
  187. )
  188. # short circuit if any is unsafe
  189. if not k_is_safe:
  190. return False
  191. v_is_safe = (
  192. _check_for_seq_len_0_and_consistent_head_dim_nested_helper(
  193. params.value, "value", debug
  194. )
  195. if params.value.is_nested
  196. else True
  197. )
  198. # short circuit if any is unsafe
  199. if not v_is_safe:
  200. return False
  201. # We now know none of the inputs have ragged num_heads, so we can safely
  202. # access .size(1)
  203. q_num_heads = params.query.size(1)
  204. k_num_heads = params.key.size(1)
  205. v_num_heads = params.value.size(1)
  206. same_num_heads = q_num_heads == k_num_heads and q_num_heads == v_num_heads
  207. if not same_num_heads:
  208. if (
  209. params.query.requires_grad
  210. or params.key.requires_grad
  211. or params.value.requires_grad
  212. ):
  213. if debug:
  214. log.warning(
  215. "Both fused kernels do not support training with broadcasted NT inputs."
  216. )
  217. return False
  218. return _try_broadcast_param_size(
  219. q_num_heads, k_num_heads, v_num_heads, "num heads", debug
  220. )
  221. return True
  222. def _can_use_flash_sdpa_jagged(params: SDPAParams, debug=False) -> bool:
  223. constraints = (
  224. _check_batch_size_nested,
  225. _check_head_dim_size_flash_nested,
  226. _check_for_seq_len_0_nested,
  227. )
  228. for constraint in constraints:
  229. if not constraint(params, debug):
  230. return False
  231. return True
  232. def _can_use_efficient_sdpa_jagged(params: SDPAParams, debug=False) -> bool:
  233. constraints = (
  234. _check_batch_size_nested,
  235. _check_for_seq_len_0_nested,
  236. )
  237. for constraint in constraints:
  238. if not constraint(params, debug):
  239. return False
  240. return True
  241. def _can_use_math_sdpa_jagged(params: SDPAParams, debug=False) -> bool:
  242. if (
  243. not params.query.transpose(1, 2).is_contiguous()
  244. or not params.key.transpose(1, 2).is_contiguous()
  245. or not params.value.transpose(1, 2).is_contiguous()
  246. ):
  247. if debug:
  248. log.warning(
  249. "If inputs are nested tensors they must be contiguous after transposing."
  250. )
  251. return False
  252. if params.is_causal:
  253. if debug:
  254. log.warning(
  255. "Nested tensors for query / key are not supported when is_causal=True."
  256. )
  257. return False
  258. return True
  259. def _select_sdp_backend(query, key, value, attn_mask, dropout, is_causal, enable_gqa):
  260. if (
  261. not flash_sdp_enabled()
  262. and not mem_efficient_sdp_enabled()
  263. and not math_sdp_enabled()
  264. and not cudnn_sdp_enabled()
  265. ):
  266. return SDPBackend.ERROR
  267. ordering = (
  268. SDPBackend.FLASH_ATTENTION,
  269. SDPBackend.EFFICIENT_ATTENTION,
  270. SDPBackend.MATH,
  271. SDPBackend.CUDNN_ATTENTION,
  272. )
  273. params = SDPAParams(query, key, value, attn_mask, dropout, is_causal, enable_gqa)
  274. for backend in ordering:
  275. if backend == SDPBackend.CUDNN_ATTENTION:
  276. if can_use_cudnn_attention(params):
  277. return SDPBackend.CUDNN_ATTENTION
  278. if backend == SDPBackend.FLASH_ATTENTION:
  279. if can_use_flash_attention(params) and _can_use_flash_sdpa_jagged(params):
  280. return SDPBackend.FLASH_ATTENTION
  281. if backend == SDPBackend.EFFICIENT_ATTENTION:
  282. if can_use_efficient_attention(params) and _can_use_efficient_sdpa_jagged(
  283. params
  284. ):
  285. return SDPBackend.EFFICIENT_ATTENTION
  286. if backend == SDPBackend.MATH:
  287. if math_sdp_enabled() and _can_use_math_sdpa_jagged(params):
  288. return SDPBackend.MATH
  289. log.warning("Memory efficient kernel not used because:")
  290. can_use_efficient_attention(params, debug=True)
  291. _can_use_efficient_sdpa_jagged(params, debug=True)
  292. log.warning("Flash attention kernel not used because:")
  293. can_use_flash_attention(params, debug=True)
  294. _can_use_flash_sdpa_jagged(params, debug=True)
  295. log.warning("Math attention kernel not used because:")
  296. _can_use_math_sdpa_jagged(params, debug=True)
  297. log.warning("cuDNN attention kernel not used because:")
  298. can_use_cudnn_attention(params, debug=True)
  299. return SDPBackend.ERROR
  300. def _cumulative_and_max_seq_len_nnz(qkv: torch.Tensor) -> tuple[torch.Tensor, int, int]:
  301. # This function is used to calculate two pieces of metadata that are needed
  302. # for use with flash-attention and efficient_attention kernels. They are the
  303. # cumulative sequence_length over a batch of sequences and the maximum
  304. # sequence length.
  305. # It returns a tuple of cumulative sequence lengths and the maximum sequence
  306. # length, and the last element in the cumulative_sequence_lengths
  307. if not isinstance(qkv, NestedTensor):
  308. raise ValueError("QKV must be nested for flash cumulative_seq_len calculation.")
  309. if qkv.lengths() is None:
  310. # TODO: Explore performance impact of copying
  311. cumulative_seqlen = qkv.offsets().to(dtype=torch.int32, device=qkv.device)
  312. max_seqlen = qkv._get_max_seqlen()
  313. n_elem = qkv.values().shape[0]
  314. else:
  315. # TODO: Explore performance impact of copying
  316. cumulative_seqlen = (
  317. qkv.lengths().cumsum(0).to(dtype=torch.int32, device=qkv.device)
  318. )
  319. max_seqlen = qkv._get_max_seqlen()
  320. # TODO: Explore performance impact when compiling
  321. n_elem = int(cumulative_seqlen[-1].item())
  322. return cumulative_seqlen, max_seqlen, n_elem
  323. def _is_safe_to_get_storage_as_tensor(tensor: torch.Tensor):
  324. # This function checks if a nested tensor is valid for
  325. # use with the flash-attention and efficient_attention kernels without
  326. # needing to call contiguous on the nested tensor input.
  327. # It checks that the storage offsets' adjacent_differences are a constant
  328. # multiple of the previous tensor in the nested tensor and that the strides
  329. # are monitonically decreasing. This check is done after calling transpose on
  330. # the nested tensor resulting in a Nt of shape [bsz, {seq_len}, num_heads, dim]
  331. # Returns a boolean indicating if contiguous needs to be called for input
  332. assert isinstance(tensor, NestedTensor)
  333. offsets = tensor.offsets()
  334. strides = tensor._strides
  335. n_tensors = offsets.size(0) - 1
  336. if n_tensors <= 1:
  337. return True
  338. # Check initially that the tensor strides are in strictly descending order
  339. prev_stride = strides[1]
  340. for stride in strides[2:]:
  341. if prev_stride <= stride:
  342. # This would mean that the last stride is greater than the seq_len
  343. # stride
  344. return False
  345. prev_stride = stride
  346. # Congrats you made it!
  347. return True
  348. def _view_as_dense(
  349. tensor: torch.Tensor, Nnz: int, num_heads: int, head_dim: int
  350. ) -> torch.Tensor:
  351. if tensor.is_nested:
  352. return tensor.values()
  353. return tensor.view(Nnz, num_heads, head_dim)
  354. # TODO: Next iteration should add test cases and check it works
  355. # def _sdpa_nested_preprocessing_with_broadcast(query, key, value):
  356. # # Query (Batch x Num_heads x {Q_seq_len} x Dim_per_head)
  357. # # Key (Batch x Num_heads x {KV_seq_len} x Dim_per_head)
  358. # # Value (Batch x Num_heads x {KV_seq_len} x Dim_per_head)
  359. # q_batch_size = query.size(0)
  360. # k_batch_size = key.size(0)
  361. # v_batch_size = value.size(0)
  362. # output_batch_size = max(q_batch_size, k_batch_size, v_batch_size)
  363. # q_num_heads = query.size(1)
  364. # k_num_heads = key.size(1)
  365. # v_num_heads = value.size(1)
  366. # output_num_heads = max(q_num_heads, k_num_heads, v_num_heads)
  367. # head_dim_qk = query.size(3)
  368. # head_dim_v = value.size(3)
  369. # q_t = query.transpose(1, 2)
  370. # k_t = key.transpose(1, 2)
  371. # v_t = value.transpose(1, 2)
  372. # # Checks in sdp_utils ensure that if {*}_batch_size/{*}_num_heads !=
  373. # # output_batch_size/num_heads then they are 1
  374. # q_batch_size_needs_broadcast = q_batch_size != output_batch_size
  375. # k_batch_size_needs_broadcast = k_batch_size != output_batch_size
  376. # v_batch_size_needs_broadcast = v_batch_size != output_batch_size
  377. # # If {*}_batch_size_needs_broadcast, then
  378. # # (1) max_seqlen_batch_{*} is given by {*}_t.size(1)
  379. # # this is because needs_broadcast indicates that the batch_size is 1
  380. # # and hence there is only 1 value for seq_len
  381. # # (2) The cum_seq_lens are given by [0, {*}_t.size(1), 2 * {*}_t.size(1),
  382. # # ..., outut_batch_size * {*}_t.size(1)]
  383. # # (3) Nnz_{*} is given by output_batch_size * {*}_t.size(1)
  384. # if q_batch_size_needs_broadcast or not q_t.is_nested:
  385. # max_seqlen_batch_q = q_t.size(1)
  386. # cumulative_sequence_length_q = torch.arange(
  387. # 0,
  388. # (output_batch_size + 1) * max_seqlen_batch_q,
  389. # max_seqlen_batch_q,
  390. # device=q_t.device,
  391. # dtype=torch.int32,
  392. # )
  393. # Nnz_q = output_batch_size * max_seqlen_batch_q
  394. # else:
  395. # (
  396. # cumulative_sequence_length_q,
  397. # max_seqlen_batch_q,
  398. # Nnz_q,
  399. # ) = _cumulative_and_max_seq_len_nnz(q_t)
  400. # if k_batch_size_needs_broadcast and v_batch_size_needs_broadcast:
  401. # assert k_t.size(1) == v_t.size(1)
  402. # max_seqlen_batch_kv = k_t.size(1)
  403. # cumulative_sequence_length_kv = torch.arange(
  404. # 0,
  405. # (output_batch_size + 1) * max_seqlen_batch_kv,
  406. # max_seqlen_batch_kv,
  407. # device=k_t.device,
  408. # dtype=torch.int32,
  409. # )
  410. # Nnz_kv = output_batch_size * max_seqlen_batch_kv
  411. # else:
  412. # cumulative_sequence_length_kv, max_seqlen_batch_kv, Nnz_kv = (
  413. # _cumulative_and_max_seq_len_nnz(v_t)
  414. # if k_batch_size_needs_broadcast
  415. # else _cumulative_and_max_seq_len_nnz(k_t)
  416. # )
  417. # q_num_heads_needs_broadcast = q_num_heads != output_num_heads
  418. # k_num_heads_needs_broadcast = k_num_heads != output_num_heads
  419. # v_num_heads_needs_broadcast = v_num_heads != output_num_heads
  420. # if not q_t.is_nested:
  421. # query_buffer_reshaped = q_t.expand(
  422. # output_batch_size, q_t.size(1), output_num_heads, head_dim_qk
  423. # )
  424. # query_buffer_reshaped = query_buffer_reshaped.reshape(
  425. # Nnz_q, output_num_heads, head_dim_qk
  426. # )
  427. # else:
  428. # if not q_t.is_contiguous() and not _is_safe_to_get_storage_as_tensor(q_t):
  429. # q_t = q_t.contiguous()
  430. # # If we are broadcasting then Nnz_q will be the output_batch_size since
  431. # # seq_len is 1
  432. # effective_batch_size_q = (
  433. # output_batch_size if q_batch_size_needs_broadcast else Nnz_q
  434. # )
  435. # query_buffer_reshaped = _view_as_dense(
  436. # q_t, effective_batch_size_q, output_num_heads, head_dim_qk
  437. # )
  438. # # If the physical layout of the NestedTensor's storage
  439. # # is not: batch, {seq_len}, num_heads, head_dim then we need
  440. # # to call contiguous
  441. # if not k_t.is_contiguous() and not _is_safe_to_get_storage_as_tensor(k_t):
  442. # k_t = k_t.contiguous()
  443. # if not v_t.is_contiguous() and not _is_safe_to_get_storage_as_tensor(v_t):
  444. # v_t = v_t.contiguous()
  445. # effective_batch_size_k = (
  446. # output_batch_size if k_batch_size_needs_broadcast else Nnz_kv
  447. # )
  448. # key_buffer_reshaped = _view_as_dense(
  449. # k_t, effective_batch_size_k, output_num_heads, head_dim_qk
  450. # )
  451. # effective_batch_size_v = (
  452. # output_batch_size if v_batch_size_needs_broadcast else Nnz_kv
  453. # )
  454. # value_buffer_reshaped = _view_as_dense(
  455. # v_t, effective_batch_size_v, output_num_heads, head_dim_v
  456. # )
  457. # if not q_batch_size_needs_broadcast:
  458. # output_shape = q_t._size
  459. # if head_dim_v != head_dim_qk:
  460. # output_shape[-1] = head_dim_v
  461. # if q_num_heads_needs_broadcast:
  462. # output_shape[1] = output_num_heads
  463. # else:
  464. # output_shape = torch.empty(3, dtype=torch.int64, device=torch.device("cpu"))
  465. # output_shape[0] = q_t.size(1)
  466. # output_shape[1] = output_num_heads
  467. # output_shape[2] = head_dim_v
  468. # return (
  469. # query_buffer_reshaped,
  470. # key_buffer_reshaped,
  471. # value_buffer_reshaped,
  472. # cumulative_sequence_length_q,
  473. # cumulative_sequence_length_kv,
  474. # max_seqlen_batch_q,
  475. # max_seqlen_batch_kv,
  476. # output_shape,
  477. # )
  478. def _sdpa_nested_preprocessing(query, key, value):
  479. # Query (Batch x Num_heads x {Q_seq_len} x Dim_per_head)
  480. # Key (Batch x Num_heads x {KV_seq_len} x Dim_per_head)
  481. # Value (Batch x Num_heads x {KV_seq_len} x Dim_per_head)
  482. q_batch_size = query.size(0)
  483. k_batch_size = key.size(0)
  484. v_batch_size = value.size(0)
  485. q_num_heads = query.size(1)
  486. k_num_heads = key.size(1)
  487. v_num_heads = value.size(1)
  488. if not (q_batch_size == k_batch_size and q_batch_size == v_batch_size) or not (
  489. q_num_heads == k_num_heads and k_num_heads == v_num_heads
  490. ):
  491. raise RuntimeError(
  492. "This path is currently not implemented for jagged layout NT."
  493. )
  494. # return _sdpa_nested_preprocessing_with_broadcast(query, key, value)
  495. num_heads = query.size(1)
  496. head_dim_qk = query.size(3)
  497. head_dim_v = value.size(3)
  498. q_t = query.transpose(1, 2)
  499. k_t = key.transpose(1, 2)
  500. v_t = value.transpose(1, 2)
  501. (
  502. cumulative_sequence_length_q,
  503. max_seqlen_batch_q,
  504. Nnz_q,
  505. ) = _cumulative_and_max_seq_len_nnz(q_t)
  506. (
  507. cumulative_sequence_length_kv,
  508. max_seqlen_batch_kv,
  509. Nnz_kv,
  510. ) = _cumulative_and_max_seq_len_nnz(k_t)
  511. # [TODO] K and V have to have the same Nnz, should probably torch_check
  512. # assume in order to not iterate over v
  513. # If the physical layout of the NestedTensor's storage
  514. # is not: batch, {seq_len}, num_heads, head_dim then we need
  515. # to call contiguous
  516. if not q_t.is_contiguous() and not _is_safe_to_get_storage_as_tensor(q_t):
  517. q_t = q_t.contiguous()
  518. if not k_t.is_contiguous() and not _is_safe_to_get_storage_as_tensor(k_t):
  519. k_t = k_t.contiguous()
  520. if not v_t.is_contiguous() and not _is_safe_to_get_storage_as_tensor(v_t):
  521. v_t = v_t.contiguous()
  522. query_buffer_reshaped = _view_as_dense(q_t, Nnz_q, num_heads, head_dim_qk)
  523. key_buffer_reshaped = _view_as_dense(k_t, Nnz_kv, num_heads, head_dim_qk)
  524. value_buffer_reshaped = _view_as_dense(v_t, Nnz_kv, num_heads, head_dim_v)
  525. output_nt_info = {
  526. "offsets": q_t.offsets(),
  527. "lengths": q_t.lengths(),
  528. "max_seqlen": q_t._get_max_seqlen(),
  529. "min_seqlen": q_t._get_min_seqlen(),
  530. }
  531. return (
  532. query_buffer_reshaped,
  533. key_buffer_reshaped,
  534. value_buffer_reshaped,
  535. cumulative_sequence_length_q,
  536. cumulative_sequence_length_kv,
  537. max_seqlen_batch_q,
  538. max_seqlen_batch_kv,
  539. output_nt_info,
  540. )
  541. def _pad_last_dim(
  542. tensor: torch.Tensor, alignment_size: int, slice: bool
  543. ) -> torch.Tensor:
  544. # FlashAttentionV2 requires that head dimension be a multiple of 8
  545. # This was previously done within the kernel, however
  546. # This causes the kernel to maybe alias query, key, value
  547. # So instead we pad the head_dimensions to be a multiple of 8
  548. # in the composite region
  549. last_dim_size = tensor.size(-1)
  550. if last_dim_size % alignment_size == 0:
  551. return tensor
  552. pad_count = alignment_size - (last_dim_size % alignment_size)
  553. tensor = torch.nn.functional.pad(tensor, [0, pad_count])
  554. if slice:
  555. return tensor[..., 0:last_dim_size]
  556. return tensor
  557. # TODO: coalesce with torch/nn/utils/attention.py
  558. def _calculate_scale(query, scale):
  559. # TODO: Investigate why math.sqrt() isn't properly handled by Dynamo?
  560. softmax_scale = scale if scale is not None else torch.sym_sqrt(1.0 / query.size(-1))
  561. return softmax_scale
  562. def _post_process_flash_output(out: torch.Tensor, og_size):
  563. if not out.is_nested and out.size(-1) != og_size:
  564. out = out[..., 0:og_size]
  565. return out
  566. def _is_computing_meta_flops(x):
  567. # Note: there's a use case of using meta tensors & the dispatch-based flop counter.
  568. # We can use this function to check for this scenario in order to handle it specially.
  569. if not torch.jit.is_scripting() and x.device.type == "meta":
  570. torch_dispatch_mode_stack = (
  571. torch.utils._python_dispatch._get_current_dispatch_mode_stack()
  572. )
  573. return any(
  574. type(x) == torch.utils.flop_counter._FlopCounterMode
  575. for x in torch_dispatch_mode_stack
  576. )
  577. return False
  578. def _autocast(
  579. query: torch.Tensor,
  580. key: torch.Tensor,
  581. value: torch.Tensor,
  582. attn_mask: Optional[torch.Tensor],
  583. ) -> tuple[torch.Tensor, torch.Tensor, torch.Tensor, Optional[torch.Tensor]]:
  584. """
  585. [Autocasting SDPA for NJT]
  586. Normal autocasting doesn't work for NJT+SDPA right now:
  587. * NJT intercepts the __torch_function__ call for scaled_dot_product_attention, which happens
  588. before we get to any aten ops or dispatcher logic; then the torch_function logic calls into
  589. efficient attention or flash attention. So, autocasting on the scaled_dot_product_attention
  590. op won't work because we never see that aten op.
  591. * If we put autocasting on `_flash_attention_forward`, then we'll get autocasting to run, but
  592. the kernel selection logic in torch_function handling (ie. jagged_scaled_dot_product_attention)
  593. won't work correctly: the kernel selection logic will run before autocasting, and choose
  594. a kernel based on the un-autocasted dtypes; but then autocasting will run and the actual
  595. attention computation will happen in a different dtype.
  596. An alternative is to just change the backend selection logic for SDPA+NJT to be autocast-aware
  597. and rely on autocasting to do the actual conversions for flash attention / efficient attention.
  598. However, by manually doing the actual autocast before the backend selection, we ensure that the
  599. autocast handling for backend selection doesn't diverge from the autocast handling for the
  600. actual dtype conversions.
  601. """
  602. device_type = query.device.type
  603. # meta device is not supported by autocast, so break early for it
  604. if _is_computing_meta_flops(query) or not torch.is_autocast_enabled(device_type):
  605. return query, key, value, attn_mask
  606. def cvt(x):
  607. if x is None:
  608. return x
  609. target_dtype = torch.get_autocast_dtype(device_type)
  610. if (
  611. (not x.dtype.is_floating_point)
  612. or x.dtype == target_dtype
  613. or x.dtype == torch.float64
  614. ):
  615. return x
  616. return x.to(target_dtype)
  617. return cvt(query), cvt(key), cvt(value), cvt(attn_mask)
  618. def jagged_scaled_dot_product_attention(
  619. query: torch.Tensor,
  620. key: torch.Tensor,
  621. value: torch.Tensor,
  622. attn_mask: Optional[torch.Tensor] = None,
  623. dropout_p=0.0,
  624. is_causal=False,
  625. scale=None,
  626. enable_gqa=False,
  627. ):
  628. query, key, value, attn_mask = _autocast(query, key, value, attn_mask)
  629. _validate_sdpa_input(query, key, value, attn_mask, dropout_p, is_causal, scale)
  630. # for mypy, ugh
  631. assert (
  632. isinstance(query, NestedTensor)
  633. and isinstance(key, NestedTensor)
  634. and isinstance(value, NestedTensor)
  635. )
  636. from torch.nested._internal.nested_tensor import (
  637. nested_view_from_values_offsets_lengths,
  638. )
  639. # Special path for non-ragged sequence length (e.g. for SAM where we have a ragged
  640. # second batch dim instead). For this case, we can just send the dense buffers through
  641. # vanilla SDPA.
  642. if query.dim() > 3 and key.dim() > 3 and value.dim() > 3 and query._ragged_idx == 1:
  643. output = F.scaled_dot_product_attention(
  644. query.values(),
  645. key.values(),
  646. value.values(),
  647. attn_mask=(
  648. attn_mask.values() if isinstance(attn_mask, NestedTensor) else attn_mask
  649. ),
  650. dropout_p=dropout_p,
  651. is_causal=is_causal,
  652. scale=scale,
  653. )
  654. return nested_view_from_values_offsets_lengths(
  655. output,
  656. query.offsets(),
  657. query.lengths(),
  658. min_seqlen=query._maybe_min_seqlen, # type: ignore[attr-defined]
  659. max_seqlen=query._maybe_max_seqlen, # type: ignore[attr-defined]
  660. )
  661. compute_logsumexp = query.requires_grad or key.requires_grad or value.requires_grad
  662. backend_choice = _select_sdp_backend(
  663. query, key, value, attn_mask, dropout_p, is_causal, enable_gqa
  664. )
  665. if _is_computing_meta_flops(query):
  666. # Backend choice will probably not be correct if we have a meta device,
  667. # because backend choice is device-aware. In this case, we mostly just
  668. # want to avoid using math backend (which does a .item() call).
  669. # Arbitrarily choose flash attention.
  670. backend_choice = SDPBackend.FLASH_ATTENTION
  671. if backend_choice == SDPBackend.FLASH_ATTENTION:
  672. og_size = query.size(-1)
  673. query_padded = _pad_last_dim(query, 8, False)
  674. key_padded = _pad_last_dim(key, 8, False)
  675. value_padded = _pad_last_dim(value, 8, False)
  676. # We need to calculate the scale based off the OG head dim size
  677. og_scale = _calculate_scale(query, scale)
  678. (
  679. query_buffer_reshaped,
  680. key_buffer_reshaped,
  681. value_buffer_reshaped,
  682. cumulative_sequence_length_q,
  683. cumulative_sequence_length_kv,
  684. max_seqlen_batch_q,
  685. max_seqlen_batch_kv,
  686. output_nt_info,
  687. ) = _sdpa_nested_preprocessing(query_padded, key_padded, value_padded)
  688. (
  689. attention,
  690. _logsumexp,
  691. _philox_seed,
  692. _philox_offset,
  693. _debug_attn_mask,
  694. ) = torch.ops.aten._flash_attention_forward(
  695. query_buffer_reshaped,
  696. key_buffer_reshaped,
  697. value_buffer_reshaped,
  698. cumulative_sequence_length_q,
  699. cumulative_sequence_length_kv,
  700. max_seqlen_batch_q,
  701. max_seqlen_batch_kv,
  702. dropout_p,
  703. is_causal,
  704. False,
  705. scale=og_scale,
  706. )
  707. # Reshape output to convert nnz to batch_size and seq_len
  708. attention = nested_view_from_values_offsets_lengths(
  709. attention, # output from flash_attn is [total_q, num_heads, head_size_og]
  710. **output_nt_info,
  711. ).transpose(1, 2)
  712. return _post_process_flash_output(attention, og_size)
  713. elif backend_choice == SDPBackend.EFFICIENT_ATTENTION:
  714. (
  715. query_reshaped,
  716. key_reshaped,
  717. value_reshaped,
  718. cumulative_sequence_length_q,
  719. cumulative_sequence_length_kv,
  720. max_seqlen_batch_q,
  721. max_seqlen_batch_kv,
  722. output_nt_info,
  723. ) = _sdpa_nested_preprocessing(query, key, value)
  724. (
  725. attention,
  726. log_sumexp,
  727. seed,
  728. offset,
  729. max_seqlen_q,
  730. max_seqlen_batch_kv,
  731. ) = torch.ops.aten._efficient_attention_forward(
  732. query_reshaped.unsqueeze(0),
  733. key_reshaped.unsqueeze(0),
  734. value_reshaped.unsqueeze(0),
  735. None,
  736. cumulative_sequence_length_q,
  737. cumulative_sequence_length_kv,
  738. max_seqlen_batch_q,
  739. max_seqlen_batch_kv,
  740. dropout_p,
  741. int(is_causal),
  742. compute_logsumexp,
  743. scale=scale,
  744. )
  745. # Reshape output to convert nnz to batch_size and seq_len
  746. return nested_view_from_values_offsets_lengths(
  747. attention.squeeze(0),
  748. **output_nt_info,
  749. ).transpose(1, 2)
  750. elif backend_choice == SDPBackend.CUDNN_ATTENTION:
  751. (
  752. query_reshaped,
  753. key_reshaped,
  754. value_reshaped,
  755. cumulative_sequence_length_q,
  756. cumulative_sequence_length_kv,
  757. max_seqlen_batch_q,
  758. max_seqlen_batch_kv,
  759. output_nt_info,
  760. ) = _sdpa_nested_preprocessing(query, key, value)
  761. (
  762. attention,
  763. logsumexp,
  764. cum_seqlen_q,
  765. cum_seqlen_kv,
  766. max_seqlen_q,
  767. max_seqlen_kv,
  768. seed,
  769. offset,
  770. _,
  771. ) = torch.ops.aten._cudnn_attention_forward(
  772. query_reshaped,
  773. key_reshaped,
  774. value_reshaped,
  775. attn_mask,
  776. cumulative_sequence_length_q,
  777. cumulative_sequence_length_kv,
  778. max_seqlen_batch_q,
  779. max_seqlen_batch_kv,
  780. compute_logsumexp,
  781. dropout_p,
  782. is_causal,
  783. False,
  784. scale=scale,
  785. )
  786. return nested_view_from_values_offsets_lengths(
  787. attention,
  788. **output_nt_info,
  789. ).transpose(1, 2)
  790. elif backend_choice == SDPBackend.MATH:
  791. # save the offsets and shape of the inputs, so we can reshape the final output
  792. # query @ key = attn: [B, D1, j0, D'] @ [B, D1, D' j1] = [B, D1, j0, j1]
  793. # attn @ value = out: [B, D1, j0, j1] @ [B, D1, j1, D2] = [B, D1, j0, D2]
  794. offsets = query.offsets()
  795. q_lengths = query.lengths()
  796. min_seqlen = query._maybe_min_seqlen
  797. max_seqlen = query._maybe_max_seqlen
  798. d1 = query._size[1]
  799. d2 = value._size[-1]
  800. # convert jagged layout Nested Tensor to strided layout Nested Tensor
  801. # which support the math implementation of SDPA
  802. def get_strided_layout_nested_tensor(jagged_layout_nt):
  803. lengths = jagged_layout_nt._offsets[1:] - jagged_layout_nt._offsets[:-1]
  804. transpose = torch.transpose(jagged_layout_nt, 1, 2)
  805. tensor_list = transpose.values().split(list(lengths), dim=0)
  806. strided_nt = torch.nested.as_nested_tensor(list(tensor_list))
  807. strided_nt = strided_nt.transpose(1, 2).contiguous()
  808. return strided_nt
  809. query = get_strided_layout_nested_tensor(query)
  810. key = get_strided_layout_nested_tensor(key)
  811. value = get_strided_layout_nested_tensor(value)
  812. attn_out = torch._scaled_dot_product_attention_math(
  813. query, key, value, attn_mask, dropout_p, is_causal, scale=scale
  814. )[0]
  815. # convert strided layout Nested Tensor back to jagged layout Nested Tensor
  816. attn_out = attn_out.transpose(1, 2).contiguous().values()
  817. attn_out = attn_out.view(-1, d1, d2)
  818. attn_out = nested_view_from_values_offsets_lengths(
  819. attn_out,
  820. offsets,
  821. lengths=q_lengths,
  822. min_seqlen=min_seqlen,
  823. max_seqlen=max_seqlen,
  824. ).transpose(1, 2)
  825. return attn_out
  826. else:
  827. raise RuntimeError(
  828. "No viable backend for scaled_dot_product_attention was found."
  829. )