flash_bwd_preprocess_kernel.h 12 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254
  1. /******************************************************************************
  2. * Copyright (c) 2024, Jay Shah, Ganesh Bikshandi, Ying Zhang, Vijay Thakkar, Pradeep Ramani, Tri Dao.
  3. ******************************************************************************/
  4. #pragma once
  5. #include "cute/tensor.hpp"
  6. #include <cutlass/cutlass.h>
  7. #include <cutlass/array.h>
  8. #include <cutlass/numeric_types.h>
  9. #include <cutlass/numeric_conversion.h>
  10. #include "utils.h"
  11. namespace flash {
  12. using namespace cute;
  13. template <class TileShape_MK_, class Element, class ElementAccum, class ArchTag_, bool Clear_dQaccum, bool Varlen>
  14. class FlashAttnBwdPreprocess {
  15. public:
  16. // Type Aliases
  17. using TileShape_MK = TileShape_MK_;
  18. using ArchTag = ArchTag_;
  19. static_assert(std::is_same_v<Element, cutlass::half_t> && ArchTag::kMinComputeCapability >= 75 ||
  20. std::is_same_v<Element, cutlass::bfloat16_t> && ArchTag::kMinComputeCapability >= 80 ||
  21. std::is_same_v<Element, cutlass::float_e4m3_t> && ArchTag::kMinComputeCapability >= 89);
  22. static constexpr uint32_t MaxThreadsPerBlock = 256;
  23. static constexpr uint32_t MinBlocksPerMultiprocessor = 2;
  24. static constexpr int SharedStorageSize = 0;
  25. static constexpr int kGmemElemsPerLoad = sizeof(cute::uint128_t) / sizeof(Element);
  26. static_assert(get<1>(TileShape_MK{}) % kGmemElemsPerLoad == 0, "Headdim must be a multiple of kGmemElemsPerLoad");
  27. static constexpr int kHeadDim = get<1>(TileShape_MK{});
  28. // We want kBlockKGmem to be a power of 2 so that when we do the summing,
  29. // it's just between threads in the same warp
  30. static constexpr int kBlockKGmem = kHeadDim % 128 == 0 ? 128 : (kHeadDim % 64 == 0 ? 64 : 32);
  31. static constexpr int kGmemThreadsPerRow = kBlockKGmem / kGmemElemsPerLoad;
  32. static_assert(MaxThreadsPerBlock % kGmemThreadsPerRow == 0, "MaxThreadsPerBlock must be a multiple of kGmemThreadsPerRow");
  33. using GmemLayoutAtom = Layout<Shape <Int<MaxThreadsPerBlock / kGmemThreadsPerRow>, Int<kGmemThreadsPerRow>>,
  34. Stride<Int<kGmemThreadsPerRow>, _1>>;
  35. using GmemTiledCopy = decltype(
  36. make_tiled_copy(Copy_Atom<DefaultCopy, Element>{},
  37. GmemLayoutAtom{},
  38. Layout<Shape<_1, Int<kGmemElemsPerLoad>>>{})); // Val layout, 8 or 16 vals per load
  39. static constexpr int kGmemElemsPerLoadAccum = sizeof(cute::uint128_t) / sizeof(ElementAccum);
  40. static_assert(get<1>(TileShape_MK{}) % kGmemElemsPerLoadAccum == 0, "Headdim must be a multiple of kGmemElemsPerLoadAccum");
  41. static constexpr int kGmemThreadsPerRowAccum = kBlockKGmem / kGmemElemsPerLoadAccum;
  42. static_assert(MaxThreadsPerBlock % kGmemThreadsPerRowAccum == 0, "MaxThreadsPerBlock must be a multiple of kGmemThreadsPerRowAccum");
  43. using GmemLayoutAtomAccum = Layout<Shape <Int<MaxThreadsPerBlock / kGmemThreadsPerRowAccum>, Int<kGmemThreadsPerRowAccum>>,
  44. Stride<Int<kGmemThreadsPerRowAccum>, _1>>;
  45. using GmemTiledCopyAccum = decltype(
  46. make_tiled_copy(Copy_Atom<DefaultCopy, ElementAccum>{},
  47. GmemLayoutAtomAccum{},
  48. Layout<Shape<_1, Int<kGmemElemsPerLoadAccum>>>{})); // Val layout, 4 vals per store
  49. using ShapeO = cute::Shape<int32_t, int32_t, int32_t, int32_t>; // (seqlen_q, d, head, batch)
  50. using StrideO = cute::Stride<int64_t, _1, int64_t, int64_t>;
  51. using ShapedPsum = cute::Shape<int32_t, int32_t, int32_t>; // (seqlen_q, head, batch)
  52. using StridedPsum = cute::Stride<_1, int64_t, int64_t>;
  53. // Device side arguments
  54. struct Arguments {
  55. Element const* ptr_O;
  56. ShapeO const shape_O;
  57. StrideO const stride_O;
  58. Element const* ptr_dO;
  59. StrideO const stride_dO;
  60. float* ptr_dPsum;
  61. ShapedPsum const shape_dPsum;
  62. StridedPsum const stride_dPsum;
  63. float const* ptr_LSE;
  64. StridedPsum const stride_LSE;
  65. float *ptr_LSE_log2;
  66. StridedPsum const stride_LSE_log2;
  67. ElementAccum* ptr_dQaccum;
  68. ShapeO const shape_dQaccum;
  69. StrideO const stride_dQaccum;
  70. int num_batch; // We need this to know the size of dq_semaphore in case of varlen
  71. int* dq_semaphore;
  72. int const* cu_seqlens = nullptr;
  73. int const* seqused = nullptr;
  74. };
  75. // Kernel entry point API
  76. struct Params {
  77. Element const* ptr_O;
  78. ShapeO const shape_O;
  79. StrideO const stride_O;
  80. Element const* ptr_dO;
  81. StrideO const stride_dO;
  82. float* ptr_dPsum;
  83. ShapedPsum const shape_dPsum;
  84. StridedPsum const stride_dPsum;
  85. float const* ptr_LSE;
  86. StridedPsum const stride_LSE;
  87. float* ptr_LSE_log2;
  88. StridedPsum const stride_LSE_log2;
  89. ElementAccum* ptr_dQaccum;
  90. ShapeO const shape_dQaccum;
  91. StrideO const stride_dQaccum;
  92. int num_batch;
  93. int* dq_semaphore;
  94. int const* cu_seqlens = nullptr;
  95. int const* seqused = nullptr;
  96. };
  97. // Convert to underlying arguments. In this case, a simple copy for the aliased type.
  98. static
  99. Params
  100. to_underlying_arguments(Arguments const& args) {
  101. return {
  102. args.ptr_O,
  103. args.shape_O,
  104. args.stride_O,
  105. args.ptr_dO,
  106. args.stride_dO,
  107. args.ptr_dPsum,
  108. args.shape_dPsum,
  109. args.stride_dPsum,
  110. args.ptr_LSE,
  111. args.stride_LSE,
  112. args.ptr_LSE_log2,
  113. args.stride_LSE_log2,
  114. args.ptr_dQaccum,
  115. args.shape_dQaccum,
  116. args.stride_dQaccum,
  117. args.num_batch,
  118. args.dq_semaphore,
  119. args.cu_seqlens,
  120. args.seqused
  121. };
  122. }
  123. CUTLASS_DEVICE
  124. void
  125. operator()(Params const& params, [[maybe_unused]] char* smem_buf) {
  126. static constexpr int kBlockM = get<0>(TileShape_MK{});
  127. int const thread_idx = threadIdx.x;
  128. int const m_block = blockIdx.x;
  129. int const bidh = blockIdx.y;
  130. int const bidb = blockIdx.z;
  131. bool const is_varlen = Varlen && params.cu_seqlens != nullptr;
  132. int const offset_o = !is_varlen ? 0 : params.cu_seqlens[bidb];
  133. int const seqlen_o = !is_varlen ? get<0>(params.shape_O) : (params.seqused ? params.seqused[bidb] : params.cu_seqlens[bidb + 1] - offset_o);
  134. if (is_varlen && m_block * kBlockM >= seqlen_o) { return; }
  135. Tensor mO = make_tensor(make_gmem_ptr(params.ptr_O), params.shape_O, params.stride_O)(_, _, bidh, !is_varlen ? bidb : 0);
  136. Tensor gO = local_tile(cute::domain_offset(make_coord(offset_o, _0{}), mO), TileShape_MK{}, make_coord(m_block, _0{})); // (M, K)
  137. Tensor mdO = make_tensor(make_gmem_ptr(params.ptr_dO), params.shape_O, params.stride_dO)(_, _, bidh, !is_varlen ? bidb : 0);
  138. Tensor gdO = local_tile(cute::domain_offset(make_coord(offset_o, _0{}), mdO), TileShape_MK{}, make_coord(m_block, _0{})); // (M, K)
  139. auto shape_LSE = select<0, 2, 3>(params.shape_O);
  140. Tensor mLSE = make_tensor(make_gmem_ptr(params.ptr_LSE), shape_LSE, params.stride_LSE)(_, bidh, !is_varlen ? bidb : 0);
  141. Tensor gLSE = local_tile(cute::domain_offset(make_coord(offset_o), mLSE), Shape<Int<kBlockM>>{}, make_coord(m_block));
  142. static_assert(kBlockM <= MaxThreadsPerBlock);
  143. float lse = thread_idx < seqlen_o - m_block * kBlockM && thread_idx < kBlockM ? gLSE(thread_idx) : INFINITY;
  144. GmemTiledCopy gmem_tiled_copy_O;
  145. auto gmem_thr_copy_O = gmem_tiled_copy_O.get_thread_slice(thread_idx);
  146. Tensor tOgO = gmem_thr_copy_O.partition_S(gO);
  147. Tensor tOgdO = gmem_thr_copy_O.partition_S(gdO);
  148. // Construct identity layout for gO
  149. Tensor cO = cute::make_identity_tensor(TileShape_MK{}); // (BLK_M,BLK_K) -> (blk_m,blk_k)
  150. // Repeat the partitioning with identity layouts
  151. Tensor tOcO = gmem_thr_copy_O.partition_D(cO);
  152. Tensor tOpO = make_tensor<bool>(make_shape(size<2>(tOgO)));
  153. #pragma unroll
  154. for (int k = 0; k < size(tOpO); ++k) { tOpO(k) = get<1>(tOcO(_0{}, _0{}, k)) < get<1>(params.shape_O); }
  155. // (8, kBlockM / 32, kHeadDim / 64) or (8, kBlockM / 16, kHeadDim / 128)
  156. Tensor tOrO = make_fragment_like(tOgO);
  157. Tensor tOrdO = make_fragment_like(tOgdO);
  158. flash::copy</*Is_even_MN=*/false, /*Is_even_K=*/false, /*Clear_OOB_MN=*/true>(
  159. gmem_tiled_copy_O, tOgO, tOrO, tOcO, tOpO, seqlen_o - m_block * kBlockM
  160. );
  161. flash::copy</*Is_even_MN=*/false, /*Is_even_K=*/false, /*Clear_OOB_MN=*/true>(
  162. gmem_tiled_copy_O, tOgdO, tOrdO, tOcO, tOpO, seqlen_o - m_block * kBlockM
  163. );
  164. // Reshape from e.g. (8, kBlockM / 32, kHeadDim / 64) to (kBlockM / 32, (8, kHeadDim / 64))
  165. Layout l = make_layout(get<1>(tOrO.layout()), make_layout(get<0>(tOrO.layout()), get<2>(tOrO.layout())));
  166. Tensor o_fp32 = flash::convert_type<float>(make_tensor(tOrO.data(), l));
  167. Tensor do_fp32 = flash::convert_type<float>(make_tensor(tOrdO.data(), l));
  168. // Sum across the last dimension
  169. Tensor dP_sum = make_tensor<float>(make_shape(size<0>(o_fp32)));
  170. #pragma unroll
  171. for (int mi = 0; mi < size<0>(o_fp32); ++mi) {
  172. float dP_sum_cur = do_fp32(mi, 0) * o_fp32(mi, 0);
  173. #pragma unroll
  174. for (int ni = 1; ni < size<1>(o_fp32); ni++) {
  175. dP_sum_cur += do_fp32(mi, ni) * o_fp32(mi, ni);
  176. }
  177. flash::SumOp<float> sum_op;
  178. dP_sum(mi) = flash::Allreduce<kGmemThreadsPerRow>::run(dP_sum_cur, sum_op);
  179. }
  180. // If varlen, the layout for dPSum, LSE_log2, and dQaccum is that we pad each sequence in the batch
  181. // by an extra 128, so that the write for each sequence doesn't touch the next sequence.
  182. // Sequence i starts at params.cu_seqlens[i] + i * 128 and ends at params.cu_seqlens[i + 1] + i * 128
  183. int const offset_padded = !is_varlen ? 0 : (params.cu_seqlens[bidb] + bidb * 128) / 128 * 128;
  184. Tensor mdPsum = make_tensor(make_gmem_ptr(params.ptr_dPsum), params.shape_dPsum, params.stride_dPsum)(_, bidh, !is_varlen ? bidb : 0);
  185. Tensor gdPsum = local_tile(cute::domain_offset(make_coord(offset_padded), mdPsum), Shape<Int<kBlockM>>{}, make_coord(m_block));
  186. if (thread_idx % kGmemThreadsPerRow == 0) {
  187. #pragma unroll
  188. for (int mi = 0; mi < size(dP_sum); ++mi) {
  189. int row = thread_idx / kGmemThreadsPerRow + mi * MaxThreadsPerBlock / kGmemThreadsPerRow;
  190. gdPsum(row) = row < seqlen_o - m_block * kBlockM ? dP_sum(mi) : 0;
  191. }
  192. }
  193. int const seqlen_rounded = cute::round_up(seqlen_o, kBlockM);
  194. Tensor mLSElog2 = make_tensor(make_gmem_ptr(params.ptr_LSE_log2), params.shape_dPsum, params.stride_LSE_log2)(_, bidh, !is_varlen ? bidb : 0);
  195. Tensor gLSElog2 = local_tile(cute::domain_offset(make_coord(offset_padded), mLSElog2), Shape<Int<kBlockM>>{}, make_coord(m_block));
  196. if (thread_idx < seqlen_rounded - m_block * kBlockM && thread_idx < kBlockM) {
  197. gLSElog2(thread_idx) = lse == -INFINITY ? 0.f : lse * float(M_LOG2E);
  198. }
  199. if constexpr (Clear_dQaccum) {
  200. Tensor mdQaccum = make_tensor(make_gmem_ptr(params.ptr_dQaccum), params.shape_dQaccum, params.stride_dQaccum)(_, _, bidh, !is_varlen ? bidb : 0);
  201. Tensor gdQaccum = local_tile(cute::domain_offset(make_coord(offset_padded, _0{}), mdQaccum), TileShape_MK{}, make_coord(m_block, _0{}));
  202. GmemTiledCopyAccum gmem_tiled_copy_dQaccum;
  203. auto gmem_thr_copy_dQaccum = gmem_tiled_copy_dQaccum.get_thread_slice(thread_idx);
  204. Tensor tdQgdQaccum = gmem_thr_copy_dQaccum.partition_D(gdQaccum);
  205. Tensor zero = make_fragment_like(tdQgdQaccum);
  206. clear(zero);
  207. // cute::copy(zero, tdQgdQaccum); // Somehow this doesn't vectorize the write
  208. #pragma unroll
  209. for (int m = 0; m < size<1>(zero); ++m) {
  210. cute::copy(zero(_, m, _), tdQgdQaccum(_, m, _));
  211. }
  212. }
  213. if (params.dq_semaphore != nullptr && thread_idx == 0) {
  214. int const num_batch = params.num_batch;
  215. int const num_head = get<2>(params.shape_dQaccum);
  216. params.dq_semaphore[bidh + bidb * num_head + m_block * num_head * num_batch] = 0;
  217. }
  218. }
  219. };
  220. } // namespace flash