atomic.h 28.2 KB
Newer Older
1
2
3
4
5
6
7
#pragma once

#ifndef __CUDACC_RTC__
#include <cuda_runtime.h>
#endif

#include <cuda/atomic>
8
#include <cuda_fp16.h>
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
#include <cutlass/numeric_types.h>

using cutlass::bfloat16_t;
using cutlass::half_t;

#define TL_DEVICE __forceinline__ __device__

template <typename T> struct normalize_atomic_type {
  using type = T;
};

template <> struct normalize_atomic_type<half_t> {
  using type = half;
};

#if (defined(__CUDA_ARCH_LIST__) && (__CUDA_ARCH_LIST__ > 750))
template <> struct normalize_atomic_type<bfloat16_t> {
  using type = __nv_bfloat16;
};
#endif

template <typename T1, typename T2> TL_DEVICE T1 cuda_cast(T2 val) {
  return T1(val);
}

template <> TL_DEVICE half cuda_cast<half, float>(float val) {
  return __float2half(val);
}

#if (defined(__CUDA_ARCH_LIST__) && (__CUDA_ARCH_LIST__ > 750))
template <> TL_DEVICE __nv_bfloat16 cuda_cast<__nv_bfloat16, float>(float val) {
  return __float2bfloat16(val);
}
#endif

template <typename T1, typename T2>
TL_DEVICE void AtomicMax(T1 &ref, T2 val,
                         int memory_order = int(cuda::memory_order_relaxed)) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  T1 *address = &ref;
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
  if constexpr (std::is_same_v<NT1, half> ||
                std::is_same_v<NT1, __nv_bfloat16>) {
    // There is no implementation of atomicMax for half and bf16 in cuda.
    // We simulate this process by atomicCAS loop.
    unsigned short *address_as_ushort =
        reinterpret_cast<unsigned short *>(address);
    unsigned short val_as_ushort = *reinterpret_cast<unsigned short *>(&val);
    unsigned short old_val_ushort = *address_as_ushort;
    while (val > *reinterpret_cast<T1 *>(&old_val_ushort)) {
      unsigned short assumed_val_ushort = old_val_ushort;
      old_val_ushort =
          atomicCAS(address_as_ushort, assumed_val_ushort, val_as_ushort);
      if (assumed_val_ushort == old_val_ushort) {
        break;
      }
    }
65
66
67
68
69
70
71
72
73
74
75
  } else {
    cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(*address);
    aref.fetch_max(cuda_cast<NT1>(val), cuda::memory_order(memory_order));
  }
}

template <typename T1, typename T2>
TL_DEVICE T1 AtomicMaxRet(T1 &ref, T2 val,
                          int memory_order = int(cuda::memory_order_relaxed)) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  T1 *address = &ref;
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
  if constexpr (std::is_same_v<NT1, half> ||
                std::is_same_v<NT1, __nv_bfloat16>) {
    unsigned short *address_as_ushort =
        reinterpret_cast<unsigned short *>(address);
    unsigned short val_as_ushort = *reinterpret_cast<unsigned short *>(&val);
    unsigned short old_val_ushort = *address_as_ushort;
    while (val > *reinterpret_cast<T1 *>(&old_val_ushort)) {
      unsigned short assumed_val_ushort = old_val_ushort;
      old_val_ushort =
          atomicCAS(address_as_ushort, assumed_val_ushort, val_as_ushort);
      if (assumed_val_ushort == old_val_ushort) {
        break;
      }
    }
    return static_cast<T1>(*reinterpret_cast<T1 *>(&old_val_ushort));
91
92
93
94
95
96
97
98
99
100
101
102
  } else {
    cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(*address);
    return static_cast<T1>(
        aref.fetch_max(cuda_cast<NT1>(val), cuda::memory_order(memory_order)));
  }
}

template <typename T1, typename T2>
TL_DEVICE void AtomicMin(T1 &ref, T2 val,
                         int memory_order = int(cuda::memory_order_relaxed)) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  T1 *address = &ref;
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
  if constexpr (std::is_same_v<NT1, half> ||
                std::is_same_v<NT1, __nv_bfloat16>) {
    // There is no implementation of atomicMin for half and bf16 in cuda.
    // We simulate this process by atomicCAS loop.
    unsigned short *address_as_ushort =
        reinterpret_cast<unsigned short *>(address);
    unsigned short val_as_ushort = *reinterpret_cast<unsigned short *>(&val);
    unsigned short old_val_ushort = *address_as_ushort;
    while (val < *reinterpret_cast<T1 *>(&old_val_ushort)) {
      unsigned short assumed_val_ushort = old_val_ushort;
      old_val_ushort =
          atomicCAS(address_as_ushort, assumed_val_ushort, val_as_ushort);
      if (assumed_val_ushort == old_val_ushort) {
        break;
      }
    }
119
120
121
122
123
124
125
126
127
128
129
  } else {
    cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(*address);
    aref.fetch_min(cuda_cast<NT1>(val), cuda::memory_order(memory_order));
  }
}

template <typename T1, typename T2>
TL_DEVICE T1 AtomicMinRet(T1 &ref, T2 val,
                          int memory_order = int(cuda::memory_order_relaxed)) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  T1 *address = &ref;
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
  if constexpr (std::is_same_v<NT1, half> ||
                std::is_same_v<NT1, __nv_bfloat16>) {
    unsigned short *address_as_ushort =
        reinterpret_cast<unsigned short *>(address);
    unsigned short val_as_ushort = *reinterpret_cast<unsigned short *>(&val);
    unsigned short old_val_ushort = *address_as_ushort;
    while (val < *reinterpret_cast<T1 *>(&old_val_ushort)) {
      unsigned short assumed_val_ushort = old_val_ushort;
      old_val_ushort =
          atomicCAS(address_as_ushort, assumed_val_ushort, val_as_ushort);
      if (assumed_val_ushort == old_val_ushort) {
        break;
      }
    }
    return static_cast<T1>(*reinterpret_cast<T1 *>(&old_val_ushort));
145
146
147
148
149
150
151
152
153
154
155
156
  } else {
    cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(*address);
    return static_cast<T1>(
        aref.fetch_min(cuda_cast<NT1>(val), cuda::memory_order(memory_order)));
  }
}

template <typename T1, typename T2>
TL_DEVICE void AtomicAdd(T1 &ref, T2 val,
                         int memory_order = int(cuda::memory_order_relaxed)) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  T1 *address = &ref;
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
  if constexpr (std::is_same_v<NT1, half> ||
                std::is_same_v<NT1, __nv_bfloat16>) {
    if (memory_order == int(cuda::memory_order_relaxed)) {
      atomicAdd(reinterpret_cast<NT1 *>(address), static_cast<NT1>(val));
    } else {
      // Since atomic ref do not support memory order, we need to inline ptx
      // code here for each situation
      if constexpr (std::is_same_v<NT1, half>) {
        // fp16
        __half ret_val;
        unsigned short ret_val_cast =
            *reinterpret_cast<unsigned short *>(&ret_val);
        unsigned long long ref_address =
            reinterpret_cast<unsigned long long>(address);
        unsigned short val_cast = *reinterpret_cast<unsigned short *>(&val);
        if (memory_order == int(cuda::memory_order_release) ||
            memory_order == int(cuda::memory_order_consume)) {
          asm volatile("atom.release.gpu.global.add.noftz.f16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acquire)) {
          asm volatile("atom.acquire.gpu.global.add.noftz.f16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acq_rel) ||
                   memory_order == int(cuda::memory_order_seq_cst)) {
          asm volatile("atom.acq_rel.gpu.global.add.noftz.f16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        }
      } else if constexpr (std::is_same_v<NT1, __nv_bfloat16>) {
        // bf16
        __nv_bfloat16 ret_val;
        unsigned short ret_val_cast =
            *reinterpret_cast<unsigned short *>(&ret_val);
        unsigned long long ref_address =
            reinterpret_cast<unsigned long long>(address);
        unsigned short val_cast = *reinterpret_cast<unsigned short *>(&val);
        if (memory_order == int(cuda::memory_order_release) ||
            memory_order == int(cuda::memory_order_consume)) {
          asm volatile("atom.release.gpu.global.add.noftz.bf16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acquire)) {
          asm volatile("atom.acquire.gpu.global.add.noftz.bf16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acq_rel) ||
                   memory_order == int(cuda::memory_order_seq_cst)) {
          asm volatile("atom.acq_rel.gpu.global.add.noftz.bf16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        }
      }
    }
218
219
220
221
222
223
224
225
226
227
228
  } else {
    cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(*address);
    aref.fetch_add(cuda_cast<NT1>(val), cuda::memory_order(memory_order));
  }
}

template <typename T1, typename T2>
TL_DEVICE T1 AtomicAddRet(T1 &ref, T2 val,
                          int memory_order = int(cuda::memory_order_relaxed)) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  T1 *address = &ref;
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
  if constexpr (std::is_same_v<NT1, half> ||
                std::is_same_v<NT1, __nv_bfloat16>) {
    if (memory_order == int(cuda::memory_order_relaxed)) {
      return static_cast<T1>(
          atomicAdd(reinterpret_cast<NT1 *>(address), static_cast<NT1>(val)));
    } else {
      if constexpr (std::is_same_v<NT1, half>) {
        // fp16
        __half ret_val;
        unsigned short ret_val_cast =
            *reinterpret_cast<unsigned short *>(&ret_val);
        unsigned long long ref_address =
            reinterpret_cast<unsigned long long>(address);
        unsigned short val_cast = *reinterpret_cast<unsigned short *>(&val);
        if (memory_order == int(cuda::memory_order_release) ||
            memory_order == int(cuda::memory_order_consume)) {
          asm volatile("atom.release.gpu.global.add.noftz.f16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acquire)) {
          asm volatile("atom.acquire.gpu.global.add.noftz.f16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acq_rel) ||
                   memory_order == int(cuda::memory_order_seq_cst)) {
          asm volatile("atom.acq_rel.gpu.global.add.noftz.f16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        }
        return static_cast<T1>(*reinterpret_cast<__half *>(&ret_val_cast));
      } else if constexpr (std::is_same_v<NT1, __nv_bfloat16>) {
        // bf16
        __nv_bfloat16 ret_val;
        unsigned short ret_val_cast =
            *reinterpret_cast<unsigned short *>(&ret_val);
        unsigned long long ref_address =
            reinterpret_cast<unsigned long long>(address);
        unsigned short val_cast = *reinterpret_cast<unsigned short *>(&val);
        if (memory_order == int(cuda::memory_order_release) ||
            memory_order == int(cuda::memory_order_consume)) {
          asm volatile("atom.release.gpu.global.add.noftz.bf16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acquire)) {
          asm volatile("atom.acquire.gpu.global.add.noftz.bf16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        } else if (memory_order == int(cuda::memory_order_acq_rel) ||
                   memory_order == int(cuda::memory_order_seq_cst)) {
          asm volatile("atom.acq_rel.gpu.global.add.noftz.bf16 %0, [%1], %2;"
                       : "=h"(ret_val_cast)
                       : "l"(ref_address), "h"(val_cast)
                       : "memory");
        }
        return static_cast<T1>(
            *reinterpret_cast<__nv_bfloat16 *>(&ret_val_cast));
      }
    }
292
293
294
295
296
297
298
  } else {
    cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(*address);
    return static_cast<T1>(
        aref.fetch_add(cuda_cast<NT1>(val), cuda::memory_order(memory_order)));
  }
}

299
300
301
// TODO add memory_order for vectorized atomic add
TL_DEVICE void AtomicAddx2(half_t *ref, half_t *val,
                           int memory_order = int(cuda::memory_order_relaxed)) {
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
  if (memory_order == int(cuda::memory_order_relaxed)) {
    atomicAdd(reinterpret_cast<half2 *>(ref),
              static_cast<half2>(*reinterpret_cast<half2 *>(val)));
  } else {
    // Since atomicAdd does not support memory order, atomic_ref does not
    // support vectorized atomic operation we can only inline ptx code here
    // Note: Vectorized atomic operations only support global space
    // Note: for 16-bit value, we need to reinterpret_cast the value to unsigned
    // short and use "h" register in assembly
    __half2 add_val = *reinterpret_cast<__half2 *>(val);
    unsigned short add_val_x_cast =
        *reinterpret_cast<unsigned short *>(&add_val.x);
    unsigned short add_val_y_cast =
        *reinterpret_cast<unsigned short *>(&add_val.y);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    __half ret_val_x, ret_val_y;
    unsigned short ret_val_x_cast =
        *reinterpret_cast<unsigned short *>(&ret_val_x);
    unsigned short ret_val_y_cast =
        *reinterpret_cast<unsigned short *>(&ret_val_y);
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile(
          "atom.release.gpu.global.add.noftz.v2.f16 {%0,%1}, [%2], {%3,%4};"
          : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
          : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
          : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile(
          "atom.acquire.gpu.global.add.noftz.v2.f16 {%0,%1}, [%2], {%3,%4};"
          : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
          : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
          : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile(
          "atom.acq_rel.gpu.global.add.noftz.v2.f16 {%0,%1}, [%2], {%3,%4};"
          : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
          : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
          : "memory");
    }
  }
344
345
}

346
347
348
TL_DEVICE half2
AtomicAddx2Ret(half_t *ref, half_t *val,
               int memory_order = int(cuda::memory_order_relaxed)) {
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
  if (memory_order == int(cuda::memory_order_relaxed)) {
    return atomicAdd(reinterpret_cast<half2 *>(ref),
                     static_cast<half2>(*reinterpret_cast<half2 *>(val)));
  } else {
    __half2 add_val = *reinterpret_cast<__half2 *>(val);
    unsigned short add_val_x_cast =
        *reinterpret_cast<unsigned short *>(&add_val.x);
    unsigned short add_val_y_cast =
        *reinterpret_cast<unsigned short *>(&add_val.y);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    __half ret_val_x, ret_val_y;
    unsigned short ret_val_x_cast =
        *reinterpret_cast<unsigned short *>(&ret_val_x);
    unsigned short ret_val_y_cast =
        *reinterpret_cast<unsigned short *>(&ret_val_y);
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile(
          "atom.release.gpu.global.add.noftz.v2.f16 {%0,%1}, [%2], {%3,%4};"
          : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
          : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
          : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile(
          "atom.acquire.gpu.global.add.noftz.v2.f16 {%0,%1}, [%2], {%3,%4};"
          : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
          : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
          : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile(
          "atom.acq_rel.gpu.global.add.noftz.v2.f16 {%0,%1}, [%2], {%3,%4};"
          : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
          : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
          : "memory");
    }
    return half2(*reinterpret_cast<__half *>(&ret_val_x_cast),
                 *reinterpret_cast<__half *>(&ret_val_y_cast));
  }
388
389
390
}

#if (defined(__CUDA_ARCH_LIST__) && (__CUDA_ARCH_LIST__ > 750))
391
392
TL_DEVICE void AtomicAddx2(bfloat16_t *ref, bfloat16_t *val,
                           int memory_order = int(cuda::memory_order_relaxed)) {
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
  if (memory_order == int(cuda::memory_order_relaxed)) {
    atomicAdd(
        reinterpret_cast<__nv_bfloat162 *>(ref),
        static_cast<__nv_bfloat162>(*reinterpret_cast<__nv_bfloat162 *>(val)));
  } else {
    __nv_bfloat162 add_val = *reinterpret_cast<__nv_bfloat162 *>(val);
    unsigned short add_val_x_cast =
        *reinterpret_cast<unsigned short *>(&add_val.x);
    unsigned short add_val_y_cast =
        *reinterpret_cast<unsigned short *>(&add_val.y);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    __nv_bfloat162 ret_val;
    unsigned short ret_val_x_cast =
        *reinterpret_cast<unsigned short *>(&ret_val.x);
    unsigned short ret_val_y_cast =
        *reinterpret_cast<unsigned short *>(&ret_val.y);
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile("atom.release.gpu.global.add.v2.bf16 {%0,%1}, [%2], {%3,%4};"
                   : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
                   : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile("atom.acquire.gpu.global.add.v2.bf16 {%0,%1}, [%2], {%3,%4};"
                   : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
                   : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile("atom.acq_rel.gpu.global.add.v2.bf16 {%0,%1}, [%2], {%3,%4};"
                   : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
                   : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
                   : "memory");
    }
  }
428
429
}

430
431
432
TL_DEVICE __nv_bfloat162
AtomicAddx2Ret(bfloat16_t *ref, bfloat16_t *val,
               int memory_order = int(cuda::memory_order_relaxed)) {
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
  if (memory_order == int(cuda::memory_order_relaxed)) {
    return atomicAdd(
        reinterpret_cast<__nv_bfloat162 *>(ref),
        static_cast<__nv_bfloat162>(*reinterpret_cast<__nv_bfloat162 *>(val)));
  } else {
    __nv_bfloat162 add_val = *reinterpret_cast<__nv_bfloat162 *>(val);
    unsigned short add_val_x_cast =
        *reinterpret_cast<unsigned short *>(&add_val.x);
    unsigned short add_val_y_cast =
        *reinterpret_cast<unsigned short *>(&add_val.y);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    __nv_bfloat162 ret_val;
    unsigned short ret_val_x_cast =
        *reinterpret_cast<unsigned short *>(&ret_val.x);
    unsigned short ret_val_y_cast =
        *reinterpret_cast<unsigned short *>(&ret_val.y);
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile("atom.release.gpu.global.add.v2.bf16 {%0,%1}, [%2], {%3,%4};"
                   : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
                   : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile("atom.acquire.gpu.global.add.v2.bf16 {%0,%1}, [%2], {%3,%4};"
                   : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
                   : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile("atom.acq_rel.gpu.global.add.v2.bf16 {%0,%1}, [%2], {%3,%4};"
                   : "=h"(ret_val_x_cast), "=h"(ret_val_y_cast)
                   : "l"(ref_addr), "h"(add_val_x_cast), "h"(add_val_y_cast)
                   : "memory");
    }
    return __nv_bfloat162(*reinterpret_cast<__nv_bfloat16 *>(&ret_val_x_cast),
                          *reinterpret_cast<__nv_bfloat16 *>(&ret_val_y_cast));
  }
470
471
472
473
}
#endif

#if (defined(__CUDA_ARCH_LIST__) && (__CUDA_ARCH_LIST__ >= 900))
474
475
TL_DEVICE void AtomicAddx2(float *ref, float *val,
                           int memory_order = int(cuda::memory_order_relaxed)) {
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
501
  if (memory_order == int(cuda::memory_order_relaxed)) {
    atomicAdd(reinterpret_cast<float2 *>(ref),
              static_cast<float2>(*reinterpret_cast<float2 *>(val)));
  } else {
    float2 add_val = *reinterpret_cast<float2 *>(val);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    float2 ret_val;
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile("atom.release.gpu.global.add.v2.f32 {%0,%1}, [%2], {%3,%4};"
                   : "=f"(ret_val.x), "=f"(ret_val.y)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile("atom.acquire.gpu.global.add.v2.f32 {%0,%1}, [%2], {%3,%4};"
                   : "=f"(ret_val.x), "=f"(ret_val.y)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile("atom.acq_rel.gpu.global.add.v2.f32 {%0,%1}, [%2], {%3,%4};"
                   : "=f"(ret_val.x), "=f"(ret_val.y)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y)
                   : "memory");
    }
  }
502
503
}

504
505
506
TL_DEVICE float2
AtomicAddx2Ret(float *ref, float *val,
               int memory_order = int(cuda::memory_order_relaxed)) {
507
508
509
510
511
512
513
514
515
516
517
518
519
520
521
522
523
524
525
526
527
528
529
530
531
532
533
  if (memory_order == int(cuda::memory_order_relaxed)) {
    return atomicAdd(reinterpret_cast<float2 *>(ref),
                     static_cast<float2>(*reinterpret_cast<float2 *>(val)));
  } else {
    float2 add_val = *reinterpret_cast<float2 *>(val);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    float2 ret_val;
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile("atom.release.gpu.global.add.v2.f32 {%0,%1}, [%2], {%3,%4};"
                   : "=f"(ret_val.x), "=f"(ret_val.y)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile("atom.acquire.gpu.global.add.v2.f32 {%0,%1}, [%2], {%3,%4};"
                   : "=f"(ret_val.x), "=f"(ret_val.y)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile("atom.acq_rel.gpu.global.add.v2.f32 {%0,%1}, [%2], {%3,%4};"
                   : "=f"(ret_val.x), "=f"(ret_val.y)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y)
                   : "memory");
    }
    return ret_val;
  }
534
535
}

536
537
TL_DEVICE void AtomicAddx4(float *ref, float *val,
                           int memory_order = int(cuda::memory_order_relaxed)) {
538
539
540
541
542
543
544
545
546
547
548
549
550
551
552
553
554
555
556
557
558
559
560
561
562
563
564
565
566
567
568
569
570
571
572
573
574
575
  if (memory_order == int(cuda::memory_order_relaxed)) {
    atomicAdd(reinterpret_cast<float4 *>(ref),
              static_cast<float4>(*reinterpret_cast<float4 *>(val)));
  } else {
    // Since atomicAdd does not support memory order, atomic_ref does not
    // support vectorized atomic operation we can only inline ptx code here
    // Note: Vectorized atomic operations only support global space
    float4 add_val = *reinterpret_cast<float4 *>(val);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    float4 ret_val;
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile("atom.release.gpu.global.add.v4.f32 {%0,%1,%2,%3}, [%4], "
                   "{%5,%6,%7,%8};"
                   : "=f"(ret_val.x), "=f"(ret_val.y), "=f"(ret_val.z),
                     "=f"(ret_val.w)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y),
                     "f"(add_val.z), "f"(add_val.w)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile("atom.acquire.gpu.global.add.v4.f32 {%0,%1,%2,%3}, [%4], "
                   "{%5,%6,%7,%8};"
                   : "=f"(ret_val.x), "=f"(ret_val.y), "=f"(ret_val.z),
                     "=f"(ret_val.w)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y),
                     "f"(add_val.z), "f"(add_val.w)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile("atom.acq_rel.gpu.global.add.v4.f32 {%0,%1,%2,%3}, [%4], "
                   "{%5,%6,%7,%8};"
                   : "=f"(ret_val.x), "=f"(ret_val.y), "=f"(ret_val.z),
                     "=f"(ret_val.w)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y),
                     "f"(add_val.z), "f"(add_val.w)
                   : "memory");
    }
  }
576
577
}

578
579
580
TL_DEVICE float4
AtomicAddx4Ret(float *ref, float *val,
               int memory_order = int(cuda::memory_order_relaxed)) {
581
582
583
584
585
586
587
588
589
590
591
592
593
594
595
596
597
598
599
600
601
602
603
604
605
606
607
608
609
610
611
612
613
614
615
616
  if (memory_order == int(cuda::memory_order_relaxed)) {
    return atomicAdd(reinterpret_cast<float4 *>(ref),
                     static_cast<float4>(*reinterpret_cast<float4 *>(val)));
  } else {
    float4 add_val = *reinterpret_cast<float4 *>(val);
    unsigned long long ref_addr = reinterpret_cast<unsigned long long>(ref);
    float4 ret_val;
    if (memory_order == int(cuda::memory_order_release) ||
        memory_order == int(cuda::memory_order_consume)) {
      asm volatile("atom.global.gpu.release.add.v4.f32 {%0,%1,%2,%3}, [%4], "
                   "{%5,%6,%7,%8};"
                   : "=f"(ret_val.x), "=f"(ret_val.y), "=f"(ret_val.z),
                     "=f"(ret_val.w)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y),
                     "f"(add_val.z), "f"(add_val.w)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acquire)) {
      asm volatile("atom.global.gpu.acquire.add.v4.f32 {%0,%1,%2,%3}, [%4], "
                   "{%5,%6,%7,%8};"
                   : "=f"(ret_val.x), "=f"(ret_val.y), "=f"(ret_val.z),
                     "=f"(ret_val.w)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y),
                     "f"(add_val.z), "f"(add_val.w)
                   : "memory");
    } else if (memory_order == int(cuda::memory_order_acq_rel) ||
               memory_order == int(cuda::memory_order_seq_cst)) {
      asm volatile("atom.global.gpu.acq_rel.add.v4.f32 {%0,%1,%2,%3}, [%4], "
                   "{%5,%6,%7,%8};"
                   : "=f"(ret_val.x), "=f"(ret_val.y), "=f"(ret_val.z),
                     "=f"(ret_val.w)
                   : "l"(ref_addr), "f"(add_val.x), "f"(add_val.y),
                     "f"(add_val.z), "f"(add_val.w)
                   : "memory");
    }
    return ret_val;
  }
617
618
619
620
621
622
623
624
625
626
627
628
629
630
}
#endif

template <typename T> TL_DEVICE T AtomicLoad(T &ref, int memory_order) {
  cuda::atomic_ref<T, cuda::thread_scope_device> aref(ref);
  return aref.load(cuda::memory_order(memory_order));
}

template <typename T1, typename T2>
TL_DEVICE void AtomicStore(T1 &ref, T2 value, int memory_order) {
  using NT1 = typename normalize_atomic_type<T1>::type;
  cuda::atomic_ref<NT1, cuda::thread_scope_device> aref(ref);
  aref.store(cuda_cast<NT1>(value), cuda::memory_order(memory_order));
}