forked from NVIDIA/TensorRT-LLM
-
Notifications
You must be signed in to change notification settings - Fork 0
Expand file tree
/
Copy pathcustomAllReduceKernels.h
More file actions
195 lines (171 loc) · 6.42 KB
/
customAllReduceKernels.h
File metadata and controls
195 lines (171 loc) · 6.42 KB
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
/*
* Copyright (c) 2022-2024, NVIDIA CORPORATION. All rights reserved.
*
* Licensed under the Apache License, Version 2.0 (the "License");
* you may not use this file except in compliance with the License.
* You may obtain a copy of the License at
*
* http://www.apache.org/licenses/LICENSE-2.0
*
* Unless required by applicable law or agreed to in writing, software
* distributed under the License is distributed on an "AS IS" BASIS,
* WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
* See the License for the specific language governing permissions and
* limitations under the License.
*/
#pragma once
#include <NvInferRuntime.h>
#include <cuda_bf16.h>
#include <cuda_fp16.h>
#include <limits>
#include "tensorrt_llm/common/assert.h"
#include "tensorrt_llm/common/cudaUtils.h"
namespace tensorrt_llm::kernels
{
constexpr size_t WARP_SIZE = 32;
constexpr size_t MAX_ALL_REDUCE_BLOCKS = 24;
// Use max modules to avoid overflow and ABA problem when the block num changes for barrier_flag
// Not a perfect solution, but it has large chance that it is correct
constexpr size_t MAX_ALL_REDUCE_MODULES = std::numeric_limits<uint32_t>::max() / 6 * 6;
constexpr size_t MAX_RANKS_PER_NODE = 16;
constexpr size_t DEFAULT_BLOCK_SIZE = 512;
namespace reduce_fusion::details
{
static constexpr int kBytesPerAccess = 16;
static constexpr int kWarpSize = 32;
static constexpr int kMaxCtaSize = 1024;
static constexpr int kClusterMaxSize = 8;
static constexpr int kLamportTokenNumThreshold = 16;
static constexpr int kLamportHiddenSizeThreshold = 256;
}; // namespace reduce_fusion::details
// Warning: python definition is in tensorrt_llm/functional.py
// they must be kept in sync
enum class AllReduceStrategyType : int8_t
{
NCCL = 0,
MIN_LATENCY = 1,
UB = 2,
AUTO = 3,
ONESHOT = 4,
TWOSHOT = 5,
LOWPRECISION = 6,
MNNVL = 7,
NCCL_SYMMETRIC = 8,
};
enum class AllReduceStrategyConfig : int8_t
{
USE_MEMCPY = 1 << 0,
PUSH_MODE = 1 << 1,
};
enum class AllReduceFusionOp : int8_t
{
NONE = 0,
RESIDUAL_RMS_NORM = 1,
LAST_PROCESS_FOR_UB = 2,
RESIDUAL_RMS_PREPOST_NORM = 3,
RESIDUAL_RMS_NORM_QUANT_FP8 = 4,
RESIDUAL_RMS_NORM_QUANT_NVFP4 = 5,
RESIDUAL_RMS_NORM_OUT_QUANT_FP8 = 6,
RESIDUAL_RMS_NORM_OUT_QUANT_NVFP4 = 7,
MOE_FINALIZE_ALLREDUCE_RESIDUAL_RMS_NORM = 8,
};
inline std::ostream& operator<<(std::ostream& os, AllReduceFusionOp op)
{
switch (op)
{
case AllReduceFusionOp::NONE: os << "NONE"; break;
case AllReduceFusionOp::RESIDUAL_RMS_NORM: os << "RESIDUAL_RMS_NORM"; break;
case AllReduceFusionOp::LAST_PROCESS_FOR_UB: os << "LAST_PROCESS_FOR_UB"; break;
case AllReduceFusionOp::RESIDUAL_RMS_PREPOST_NORM: os << "RESIDUAL_RMS_PREPOST_NORM"; break;
case AllReduceFusionOp::RESIDUAL_RMS_NORM_QUANT_FP8: os << "RESIDUAL_RMS_NORM_QUANT_FP8"; break;
case AllReduceFusionOp::RESIDUAL_RMS_NORM_QUANT_NVFP4: os << "RESIDUAL_RMS_NORM_QUANT_NVFP4"; break;
case AllReduceFusionOp::RESIDUAL_RMS_NORM_OUT_QUANT_FP8: os << "RESIDUAL_RMS_NORM_OUT_QUANT_FP8"; break;
case AllReduceFusionOp::RESIDUAL_RMS_NORM_OUT_QUANT_NVFP4: os << "RESIDUAL_RMS_NORM_OUT_QUANT_NVFP4"; break;
case AllReduceFusionOp::MOE_FINALIZE_ALLREDUCE_RESIDUAL_RMS_NORM:
os << "MOE_FINALIZE_ALLREDUCE_RESIDUAL_RMS_NORM";
break;
default: os << "UNKNOWN"; break;
}
return os;
}
inline std::string toString(AllReduceFusionOp op)
{
std::ostringstream oss;
oss << op;
return oss.str();
}
inline std::ostream& operator<<(std::ostream& os, AllReduceStrategyType op)
{
switch (op)
{
case AllReduceStrategyType::NCCL: os << "NCCL"; break;
case AllReduceStrategyType::MIN_LATENCY: os << "MIN_LATENCY"; break;
case AllReduceStrategyType::UB: os << "UB"; break;
case AllReduceStrategyType::AUTO: os << "AUTO"; break;
case AllReduceStrategyType::ONESHOT: os << "ONESHOT"; break;
case AllReduceStrategyType::TWOSHOT: os << "TWOSHOT"; break;
case AllReduceStrategyType::LOWPRECISION: os << "LOWPRECISION"; break;
case AllReduceStrategyType::MNNVL: os << "MNNVL"; break;
case AllReduceStrategyType::NCCL_SYMMETRIC: os << "NCCL_SYMMETRIC"; break;
}
return os;
}
inline std::string toString(AllReduceStrategyType op)
{
std::ostringstream oss;
oss << op;
return oss.str();
}
struct AllReduceFusionParams
{
AllReduceFusionParams()
: bias_buffer(nullptr)
, residual_buffer(nullptr)
, weight_buffer(nullptr)
, weight_buffer_pre_residual_norm(nullptr)
, intermediate_buffer(nullptr)
{
}
// gemm bias
void const* bias_buffer;
// residuial add
void const* residual_buffer;
// rms norm
int hidden_size; // equal to normalized_shape
void const* weight_buffer; // norm elem-wise affine gamma
void const* weight_buffer_pre_residual_norm; // for gemma norm before residual
float eps;
// new residual
void* intermediate_buffer;
void* lamport_peer_comm_buffer_ptrs[MAX_RANKS_PER_NODE * 3];
};
struct AllReduceParams
{
size_t elts_total;
size_t elts_per_rank;
size_t elts_per_block;
size_t rank_offset;
size_t ranks_per_node;
size_t local_rank;
uint32_t* barrier_flag_ptr;
uint32_t* barrier_flag_counter_ptr;
uint32_t* peer_barrier_ptrs_in[MAX_RANKS_PER_NODE];
uint32_t* peer_barrier_ptrs_out[MAX_RANKS_PER_NODE];
void* peer_comm_buffer_ptrs[MAX_RANKS_PER_NODE * 2];
void* local_output_buffer_ptr;
void const* local_input_buffer_ptr;
AllReduceFusionParams fusion_params;
static AllReduceParams deserialize(int64_t* buffer, size_t tpSize, size_t tpRank, nvinfer1::DataType dataType,
int token_num, int hidden_size, AllReduceFusionOp op);
};
bool configurationSupported(AllReduceStrategyType algo, size_t msg_size, size_t n_ranks, nvinfer1::DataType type);
void customAllReduce(kernels::AllReduceParams& params, nvinfer1::DataType dataType, AllReduceStrategyType strat,
AllReduceStrategyConfig config, AllReduceFusionOp fusionOp, cudaStream_t stream);
void residualRmsNorm(
kernels::AllReduceParams& params, nvinfer1::DataType dataType, cudaStream_t stream, AllReduceFusionOp fusionOp);
void lamportInitialize(void* buffer, size_t size, nvinfer1::DataType dataType, cudaStream_t stream);
namespace reduce_fusion
{
bool is_lamport_supported(nvinfer1::DataType dataType, int token_num, int hidden_size);
}
} // namespace tensorrt_llm::kernels