-
Notifications
You must be signed in to change notification settings - Fork 242
/
Copy pathfp8_naive_conv.cpp
571 lines (536 loc) · 24 KB
/
fp8_naive_conv.cpp
1
2
3
4
5
6
7
8
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
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
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
218
219
220
221
222
223
224
225
226
227
228
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
292
293
294
295
296
297
298
299
300
301
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
344
345
346
347
348
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
388
389
390
391
392
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
428
429
430
431
432
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
470
471
472
473
474
475
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
502
503
504
505
506
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
534
535
536
537
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
/*******************************************************************************
*
* MIT License
*
* Copyright (c) 2020-2023 Advanced Micro Devices, Inc.
*
* Permission is hereby granted, free of charge, to any person obtaining a copy
* of this software and associated documentation files (the "Software"), to deal
* in the Software without restriction, including without limitation the rights
* to use, copy, modify, merge, publish, distribute, sublicense, and/or sell
* copies of the Software, and to permit persons to whom the Software is
* furnished to do so, subject to the following conditions:
*
* The above copyright notice and this permission notice shall be included in
*all copies or substantial portions of the Software.
*
* THE SOFTWARE IS PROVIDED "AS IS", WITHOUT WARRANTY OF ANY KIND, EXPRESS OR
* IMPLIED, INCLUDING BUT NOT LIMITED TO THE WARRANTIES OF MERCHANTABILITY,
* FITNESS FOR A PARTICULAR PURPOSE AND NONINFRINGEMENT. IN NO EVENT SHALL THE
* AUTHORS OR COPYRIGHT HOLDERS BE LIABLE FOR ANY CLAIM, DAMAGES OR OTHER
* LIABILITY, WHETHER IN AN ACTION OF CONTRACT, TORT OR OTHERWISE, ARISING FROM,
* OUT OF OR IN CONNECTION WITH THE SOFTWARE OR THE USE OR OTHER DEALINGS IN THE
* SOFTWARE.
*
*******************************************************************************/
#ifndef MIOPEN_DONT_USE_HIP_RUNTIME_HEADERS
#include <hip/hip_fp16.h>
#include <hip/hip_runtime.h>
//#include <half.hpp>
#include <hip/hip_bfloat16.h>
#endif
// Copied over from naive_conv.cpp
#ifdef __HIPCC_RTC__
#ifdef WORKAROUND_ISSUE_HIPRTC_TRUE_TYPE
/// Definitions from <cstdint>, <cmath> conflict with
/// /opt/rocm/include/hip/amd_detail/amd_hip_vector_types.h.
typedef unsigned char uint8_t;
typedef signed char int8_t;
typedef signed short int16_t;
typedef unsigned short uint16_t;
typedef float float_t;
// std::conditional requires type_traits which has a few other things
// which result in collition with amd_hip_vector_types.h
namespace std {
template <bool predicate, typename X, typename Y>
struct conditional;
template <typename X, typename Y>
struct conditional<true, X, Y>
{
using type = X;
};
template <typename X, typename Y>
struct conditional<false, X, Y>
{
using type = Y;
};
template <bool predicate, typename X, typename Y>
using conditional_t = typename conditional<predicate, X, Y>::type;
} // namespace std
#else
#include <cstdint> // int8_t, int16_t
#include <cmath> // float_t
#endif
#endif // __HIPCC_RTC__
#include <limits> // std::numeric_limits
#define MIOPEN_ENABLE_F8_DEVICE_CODE 1
#include "hip_float8.hpp"
#include "fp8_kern_types.h"
using float8 = miopen_f8::hip_f8<miopen_f8::hip_f8_type::fp8>;
using bfloat8 = miopen_f8::hip_f8<miopen_f8::hip_f8_type::bf8>;
template <typename T>
inline __device__ uint32_t draft_rng(T x, uint32_t seed)
{
int i = threadIdx.x + blockIdx.x * blockDim.x;
typedef typename std::conditional<sizeof(T) == 2, uint16_t, uint32_t>::type IT;
IT tmp = *(reinterpret_cast<IT*>(&x));
uint32_t drop_bits = uint32_t(tmp) & 0xFFFFu;
if(sizeof(tmp) == 4)
drop_bits ^= tmp >> 16;
drop_bits = ((drop_bits & 31) << 11) | (drop_bits >> 5);
drop_bits *= 0x7000149;
uint32_t rng = (drop_bits ^ 0x13371337 ^ (i * 229791) ^ seed);
return rng;
}
template <typename TI, typename TO>
inline __device__ TO cast_number(const TI input, miopen_f8::hip_f8_rounding_mode mode, uint32_t rng)
{
if(std::is_same<TI, TO>::value)
{
return input;
}
if(sizeof(TI) == sizeof(TO))
{
const auto tmp = static_cast<float>(input);
return TO{tmp, mode, rng};
}
else if(sizeof(TO) > sizeof(TI))
{
return static_cast<TO>(input);
}
else
{
return TO{input, mode, rng};
}
}
template <typename TI,
typename TW,
typename TO,
typename in_cast_type,
typename wei_cast_type,
typename TACC = double>
inline __device__ void naive_conv_fwd_nchw(const TI* __restrict__ p_in,
const TW* __restrict__ p_wei,
TO* __restrict__ p_out,
const int hi,
const int wi,
const int n,
const int k_per_group,
const int c_per_group,
const int ho,
const int wo,
const int sy,
const int sx,
const int dy,
const int dx,
const int py,
const int px,
const int fy,
const int fx,
const int group,
bool stoch,
uint32_t seed)
{
/*
* need to compute total output pixel: `group * n * k_per_group * ho * wo`.
* to distribute this workload, let one workgroup compute `ho * wo` pixel,
* hence need `group * n * k_per_group` workgroups (grid_size).
*/
int k = k_per_group * group;
int c = c_per_group * group;
int thread_length = ho * wo;
int bid = blockIdx.x;
int ik = bid % k_per_group;
int in = (bid / k_per_group) % n;
int ig = bid / (n * k_per_group);
p_in += static_cast<size_t>(in) * c * hi * wi + static_cast<size_t>(ig) * c_per_group * hi * wi;
p_wei += static_cast<size_t>(ig) * k_per_group * c_per_group * fy * fx +
static_cast<size_t>(ik) * c_per_group * fy * fx;
p_out += static_cast<size_t>(in) * k * ho * wo +
static_cast<size_t>(ig) * k_per_group * ho * wo + static_cast<size_t>(ik) * ho * wo;
for(int tid = threadIdx.x; tid < thread_length; tid += 256)
{
int iho = tid / wo;
int iwo = tid % wo;
TACC value = .0f;
for(int ic = 0; ic < c_per_group; ic++)
{
for(int iy = 0; iy < fy; iy++)
{
int valid_h = 1;
int cur_h = sy * iho - py + dy * iy;
if(cur_h < 0 || cur_h >= hi)
valid_h &= 0;
for(int ix = 0; ix < fx; ix++)
{
int valid_w = 1;
int cur_w = sx * iwo - px + dx * ix;
if(cur_w < 0 || cur_w >= wi)
valid_w &= 0;
if(valid_w & valid_h)
{
size_t i_idx = static_cast<size_t>(ic) * hi * wi +
static_cast<size_t>(cur_h) * wi + static_cast<size_t>(cur_w);
size_t f_idx = static_cast<size_t>(ic) * fy * fx +
static_cast<size_t>(iy) * fx + static_cast<size_t>(ix);
uint32_t rng1 = 0;
uint32_t rng2 = 0;
auto rnd_mode = miopen_f8::hip_f8_rounding_mode::standard;
if(stoch)
{
rng1 = draft_rng(p_in[i_idx], seed);
rng2 = draft_rng(p_in[f_idx], seed);
rnd_mode = miopen_f8::hip_f8_rounding_mode::stochastic;
}
const auto item_in = in_cast_type(p_in[i_idx], rnd_mode, rng1);
const auto item_wei = wei_cast_type(p_wei[f_idx], rnd_mode, rng2);
value += static_cast<TACC>(item_in) * static_cast<TACC>(item_wei);
}
}
}
}
size_t o_idx = static_cast<size_t>(iho) * wo + static_cast<size_t>(iwo);
// p_out[o_idx] = __float2half(static_cast<float>(value));
p_out[o_idx] = static_cast<TO>(value);
}
}
extern "C" __global__ void FWD_KERNEL_NAME(const INPUT_TYPE* __restrict__ p_in,
const WEIGHTS_TYPE* __restrict__ p_wei,
OUTPUT_TYPE* __restrict__ p_out,
int hi,
int wi,
int n,
int k_per_group,
int c_per_group,
int ho,
int wo,
int sy,
int sx,
int dy,
int dx,
int py,
int px,
int fy,
int fx,
int group,
bool stochastic,
uint32_t seed)
{
// instantiate the kernel
naive_conv_fwd_nchw<INPUT_TYPE,
WEIGHTS_TYPE,
OUTPUT_TYPE,
INPUT_CAST_TYPE,
WEIGHTS_CAST_TYPE,
ACCUMULATOR_TYPE>(p_in,
p_wei,
p_out,
hi,
wi,
n,
k_per_group,
c_per_group,
ho,
wo,
sy,
sx,
dy,
dx,
py,
px,
fy,
fx,
group,
stochastic,
seed);
}
template <typename TI,
typename TW,
typename TO,
typename wei_cast_type,
typename out_cast_type,
typename TACC = double>
inline __device__ void naive_conv_bwd_nchw(TI* __restrict__ p_in,
const TW* __restrict__ p_wei,
const TO* __restrict__ p_out,
int hi,
int wi,
int n,
int k_per_group,
int c_per_group,
int ho,
int wo,
int sy,
int sx,
int dy,
int dx,
int py,
int px,
int fy,
int fx,
int group,
bool stoch,
uint32_t seed)
{
/*
* need to compute total input pixel: `group * n * c_per_group * hi * wi`.
* to distribute this workload, let one workgroup compute `hi * wi` pixel,
* hence need `group * n * c_per_group` workgroups (grid_size).
*/
int k = k_per_group * group;
int c = c_per_group * group;
int thread_length = hi * wi;
int bid = blockIdx.x;
int ic = bid % c_per_group;
int in = (bid / c_per_group) % n;
int ig = bid / (n * c_per_group);
p_in += static_cast<size_t>(in) * c * hi * wi +
static_cast<size_t>(ig) * c_per_group * hi * wi + static_cast<size_t>(ic) * hi * wi;
p_wei += static_cast<size_t>(ig) * k_per_group * c_per_group * fy * fx +
static_cast<size_t>(ic) * fy * fx;
p_out +=
static_cast<size_t>(in) * k * ho * wo + static_cast<size_t>(ig) * k_per_group * ho * wo;
for(int tid = threadIdx.x; tid < thread_length; tid += 256)
{
int ihi = tid / wi;
int iwi = tid % wi;
TACC value = .0f;
for(int ik = 0; ik < k_per_group; ik++)
{
for(int iy = 0; iy < fy; iy++)
{
int valid_h = 1;
int cur_ho = ihi + py - dy * iy; // cur_h = sy*iho-py+dy*iy;
if(cur_ho < 0 || cur_ho % sy)
valid_h &= 0;
cur_ho /= sy;
if(cur_ho >= ho)
valid_h &= 0;
for(int ix = 0; ix < fx; ix++)
{
int valid_w = 1;
int cur_wo = iwi + px - dx * ix; // cur_w = sx*iwo-px+dx*ix;
if(cur_wo < 0 || cur_wo % sx)
valid_w &= 0;
cur_wo /= sx;
if(cur_wo >= wo)
valid_w &= 0;
if(valid_h & valid_w)
{
size_t o_idx = static_cast<size_t>(ik) * ho * wo +
static_cast<size_t>(cur_ho) * wo +
static_cast<size_t>(cur_wo);
size_t f_idx = static_cast<size_t>(ik) * c_per_group * fy * fx +
static_cast<size_t>(iy) * fx + static_cast<size_t>(ix);
uint32_t rng1 = 0;
uint32_t rng2 = 0;
auto rnd_mode = miopen_f8::hip_f8_rounding_mode::standard;
if(stoch)
{
rng1 = draft_rng(p_out[o_idx], seed);
rng2 = draft_rng(p_wei[f_idx], seed);
rnd_mode = miopen_f8::hip_f8_rounding_mode::stochastic;
}
const auto item_out = out_cast_type(p_out[o_idx], rnd_mode, rng1);
const auto item_wei = wei_cast_type(p_wei[f_idx], rnd_mode, rng2);
value += static_cast<TACC>(item_out) * static_cast<TACC>(item_wei);
}
}
}
}
size_t i_idx = static_cast<size_t>(ihi) * wi + static_cast<size_t>(iwi);
p_in[i_idx] = static_cast<TI>(value);
}
}
extern "C" __global__ void BWD_KERNEL_NAME(INPUT_TYPE* __restrict__ p_in,
const WEIGHTS_TYPE* __restrict__ p_wei,
const OUTPUT_TYPE* __restrict__ p_out,
int hi,
int wi,
int n,
int k_per_group,
int c_per_group,
int ho,
int wo,
int sy,
int sx,
int dy,
int dx,
int py,
int px,
int fy,
int fx,
int group,
bool stochastic,
uint32_t seed)
{
// instantiate template
naive_conv_bwd_nchw<INPUT_TYPE,
WEIGHTS_TYPE,
OUTPUT_TYPE,
WEIGHTS_CAST_TYPE,
OUTPUT_CAST_TYPE,
ACCUMULATOR_TYPE>(p_in,
p_wei,
p_out,
hi,
wi,
n,
k_per_group,
c_per_group,
ho,
wo,
sy,
sx,
dy,
dx,
py,
px,
fy,
fx,
group,
stochastic,
seed);
}
template <typename TI,
typename TW,
typename TO,
typename in_cast_type,
typename out_cast_type,
typename TACC = double>
inline __device__ void naive_conv_wrw_nchw(const TI* __restrict__ p_in,
TW* __restrict__ p_wei,
const TO* __restrict__ p_out,
int hi,
int wi,
int n,
int k_per_group,
int c_per_group,
int ho,
int wo,
int sy,
int sx,
int dy,
int dx,
int py,
int px,
int fy,
int fx,
int group,
bool stoch,
uint32_t seed)
{
/*
* need to compute total filter pixel: `group * k_per_group * c_per_group *
* fy * fx`.
* to distribute this workload, let one workgroup compute `c_per_group * fy
* * fx` pixel,
* hence need `group * k_per_group` workgroups (grid_size).
*/
int k = k_per_group * group;
int c = c_per_group * group;
int thread_length = c_per_group * fy * fx;
int bid = blockIdx.x;
int ik = bid % k_per_group;
int ig = bid / k_per_group;
p_in += static_cast<size_t>(ig) * c_per_group * hi * wi;
p_wei += static_cast<size_t>(ig) * k_per_group * c_per_group * fy * fx +
static_cast<size_t>(ik) * c_per_group * fy * fx;
p_out += static_cast<size_t>(ig) * k_per_group * ho * wo + static_cast<size_t>(ik) * ho * wo;
for(int tid = threadIdx.x; tid < thread_length; tid += 256)
{
int ix = tid % fx;
int iy = (tid / fx) % fy;
int ic = tid / (fx * fy);
TACC value = .0f;
for(int in = 0; in < n; in++)
{
for(int iho = 0; iho < ho; iho++)
{
int valid_h = 1;
int cur_h = sy * iho - py + dy * iy;
if(cur_h < 0 || cur_h >= hi)
valid_h &= 0;
for(int iwo = 0; iwo < wo; iwo++)
{
int valid_w = 1;
int cur_w = sx * iwo - px + dx * ix;
if(cur_w < 0 || cur_w >= wi)
valid_w &= 0;
if(valid_h & valid_w)
{
size_t i_idx = static_cast<size_t>(in) * c * hi * wi +
static_cast<size_t>(ic) * hi * wi +
static_cast<size_t>(cur_h) * wi + static_cast<size_t>(cur_w);
size_t o_idx = static_cast<size_t>(in) * k * ho * wo +
static_cast<size_t>(iho) * wo + static_cast<size_t>(iwo);
uint32_t rng1 = 0;
uint32_t rng2 = 0;
auto rnd_mode = miopen_f8::hip_f8_rounding_mode::standard;
if(stoch)
{
rng1 = draft_rng(p_in[i_idx], seed);
rng2 = draft_rng(p_out[o_idx], seed);
rnd_mode = miopen_f8::hip_f8_rounding_mode::stochastic;
}
const auto item_in = in_cast_type(p_in[i_idx], rnd_mode, rng1);
const auto item_out = out_cast_type(p_out[o_idx], rnd_mode, rng2);
value += static_cast<TACC>(item_in) * static_cast<TACC>(item_out);
}
}
}
}
size_t f_idx = static_cast<size_t>(ic) * fy * fx + static_cast<size_t>(iy) * fx +
static_cast<size_t>(ix);
p_wei[f_idx] = static_cast<TW>(value);
}
}
extern "C" __global__ void WRW_KERNEL_NAME(const INPUT_TYPE* __restrict__ p_in,
WEIGHTS_TYPE* __restrict__ p_wei,
const OUTPUT_TYPE* __restrict__ p_out,
int hi,
int wi,
int n,
int k_per_group,
int c_per_group,
int ho,
int wo,
int sy,
int sx,
int dy,
int dx,
int py,
int px,
int fy,
int fx,
int group,
bool stochastic,
uint32_t seed)
{
naive_conv_wrw_nchw<INPUT_TYPE,
WEIGHTS_TYPE,
OUTPUT_TYPE,
INPUT_CAST_TYPE,
OUTPUT_CAST_TYPE,
ACCUMULATOR_TYPE>(p_in,
p_wei,
p_out,
hi,
wi,
n,
k_per_group,
c_per_group,
ho,
wo,
sy,
sx,
dy,
dx,
py,
px,
fy,
fx,
group,
stochastic,
seed);
}