Skip to content

Commit 174bf6a

Browse files
committed
Merge 'origin/master' into hipblas
2 parents f80ce7a + 1fcdcc2 commit 174bf6a

File tree

3 files changed

+112
-66
lines changed

3 files changed

+112
-66
lines changed

CMakeLists.txt

+30-26
Original file line numberDiff line numberDiff line change
@@ -37,43 +37,45 @@ endif()
3737
#
3838

3939
# general
40-
option(LLAMA_STATIC "llama: static link libraries" OFF)
41-
option(LLAMA_NATIVE "llama: enable -march=native flag" OFF)
42-
option(LLAMA_LTO "llama: enable link time optimization" OFF)
40+
option(LLAMA_STATIC "llama: static link libraries" OFF)
41+
option(LLAMA_NATIVE "llama: enable -march=native flag" OFF)
42+
option(LLAMA_LTO "llama: enable link time optimization" OFF)
4343

4444
# debug
45-
option(LLAMA_ALL_WARNINGS "llama: enable all compiler warnings" ON)
46-
option(LLAMA_ALL_WARNINGS_3RD_PARTY "llama: enable all compiler warnings in 3rd party libs" OFF)
47-
option(LLAMA_GPROF "llama: enable gprof" OFF)
45+
option(LLAMA_ALL_WARNINGS "llama: enable all compiler warnings" ON)
46+
option(LLAMA_ALL_WARNINGS_3RD_PARTY "llama: enable all compiler warnings in 3rd party libs" OFF)
47+
option(LLAMA_GPROF "llama: enable gprof" OFF)
4848

4949
# sanitizers
50-
option(LLAMA_SANITIZE_THREAD "llama: enable thread sanitizer" OFF)
51-
option(LLAMA_SANITIZE_ADDRESS "llama: enable address sanitizer" OFF)
52-
option(LLAMA_SANITIZE_UNDEFINED "llama: enable undefined sanitizer" OFF)
50+
option(LLAMA_SANITIZE_THREAD "llama: enable thread sanitizer" OFF)
51+
option(LLAMA_SANITIZE_ADDRESS "llama: enable address sanitizer" OFF)
52+
option(LLAMA_SANITIZE_UNDEFINED "llama: enable undefined sanitizer" OFF)
5353

5454
# instruction set specific
55-
option(LLAMA_AVX "llama: enable AVX" ON)
56-
option(LLAMA_AVX2 "llama: enable AVX2" ON)
57-
option(LLAMA_AVX512 "llama: enable AVX512" OFF)
58-
option(LLAMA_AVX512_VBMI "llama: enable AVX512-VBMI" OFF)
59-
option(LLAMA_AVX512_VNNI "llama: enable AVX512-VNNI" OFF)
60-
option(LLAMA_FMA "llama: enable FMA" ON)
55+
option(LLAMA_AVX "llama: enable AVX" ON)
56+
option(LLAMA_AVX2 "llama: enable AVX2" ON)
57+
option(LLAMA_AVX512 "llama: enable AVX512" OFF)
58+
option(LLAMA_AVX512_VBMI "llama: enable AVX512-VBMI" OFF)
59+
option(LLAMA_AVX512_VNNI "llama: enable AVX512-VNNI" OFF)
60+
option(LLAMA_FMA "llama: enable FMA" ON)
6161
# in MSVC F16C is implied with AVX2/AVX512
6262
if (NOT MSVC)
63-
option(LLAMA_F16C "llama: enable F16C" ON)
63+
option(LLAMA_F16C "llama: enable F16C" ON)
6464
endif()
6565

6666
# 3rd party libs
67-
option(LLAMA_ACCELERATE "llama: enable Accelerate framework" ON)
68-
option(LLAMA_BLAS "llama: use BLAS" OFF)
69-
option(LLAMA_BLAS_VENDOR "llama: BLA_VENDOR from https://cmake.org/cmake/help/latest/module/FindBLAS.html#blas-lapack-vendors" Generic)
70-
option(LLAMA_CUBLAS "llama: use cuBLAS" OFF)
71-
option(LLAMA_CLBLAST "llama: use CLBlast" OFF)
72-
option(LLAMA_HIPBLAS "llama: use hipBLAS" OFF)
73-
74-
option(LLAMA_BUILD_TESTS "llama: build tests" ${LLAMA_STANDALONE})
75-
option(LLAMA_BUILD_EXAMPLES "llama: build examples" ${LLAMA_STANDALONE})
76-
option(LLAMA_BUILD_SERVER "llama: build server example" OFF)
67+
option(LLAMA_ACCELERATE "llama: enable Accelerate framework" ON)
68+
option(LLAMA_BLAS "llama: use BLAS" OFF)
69+
option(LLAMA_BLAS_VENDOR "llama: BLA_VENDOR from https://cmake.org/cmake/help/latest/module/FindBLAS.html#blas-lapack-vendors" Generic)
70+
option(LLAMA_CUBLAS "llama: use cuBLAS" OFF)
71+
set(LLAMA_CUDA_DMMV_X "32" CACHE STRING "llama: x stride for dmmv CUDA kernels")
72+
set(LLAMA_CUDA_DMMV_Y "1" CACHE STRING "llama: y block size for dmmv CUDA kernels")
73+
option(LLAMA_CLBLAST "llama: use CLBlast" OFF)
74+
option(LLAMA_HIPBLAS "llama: use hipBLAS" OFF)
75+
76+
option(LLAMA_BUILD_TESTS "llama: build tests" ${LLAMA_STANDALONE})
77+
option(LLAMA_BUILD_EXAMPLES "llama: build examples" ${LLAMA_STANDALONE})
78+
option(LLAMA_BUILD_SERVER "llama: build server example" OFF)
7779

7880
#
7981
# Build info header
@@ -185,6 +187,8 @@ if (LLAMA_CUBLAS)
185187
set(GGML_CUDA_SOURCES ggml-cuda.cu ggml-cuda.h)
186188

187189
add_compile_definitions(GGML_USE_CUBLAS)
190+
add_compile_definitions(GGML_CUDA_DMMV_X=${LLAMA_CUDA_DMMV_X})
191+
add_compile_definitions(GGML_CUDA_DMMV_Y=${LLAMA_CUDA_DMMV_Y})
188192

189193
if (LLAMA_STATIC)
190194
set(LLAMA_EXTRA_LIBS ${LLAMA_EXTRA_LIBS} CUDA::cudart_static CUDA::cublas_static CUDA::cublasLt_static)

Makefile

+11-1
Original file line numberDiff line numberDiff line change
@@ -133,9 +133,19 @@ ifdef LLAMA_CUBLAS
133133
OBJS += ggml-cuda.o
134134
NVCC = nvcc
135135
NVCCFLAGS = --forward-unknown-to-host-compiler -arch=native
136+
ifdef LLAMA_CUDA_DMMV_X
137+
NVCCFLAGS += -DGGML_CUDA_DMMV_X=$(LLAMA_CUDA_DMMV_X)
138+
else
139+
NVCCFLAGS += -DGGML_CUDA_DMMV_X=32
140+
endif # LLAMA_CUDA_DMMV_X
141+
ifdef LLAMA_CUDA_DMMV_Y
142+
NVCCFLAGS += -DGGML_CUDA_DMMV_Y=$(LLAMA_CUDA_DMMV_Y)
143+
else
144+
NVCCFLAGS += -DGGML_CUDA_DMMV_Y=1
145+
endif # LLAMA_CUDA_DMMV_Y
136146
ggml-cuda.o: ggml-cuda.cu ggml-cuda.h
137147
$(NVCC) $(NVCCFLAGS) $(CXXFLAGS) -Wno-pedantic -c $< -o $@
138-
endif
148+
endif # LLAMA_CUBLAS
139149
ifdef LLAMA_CLBLAST
140150
CFLAGS += -DGGML_USE_CLBLAST
141151
CXXFLAGS += -DGGML_USE_CLBLAST

ggml-cuda.cu

+71-39
Original file line numberDiff line numberDiff line change
@@ -130,9 +130,19 @@ typedef struct {
130130
} block_q8_0;
131131
static_assert(sizeof(block_q8_0) == sizeof(ggml_fp16_t) + QK8_0, "wrong q8_0 block size/padding");
132132

133+
#define WARP_SIZE 32
134+
133135
#define CUDA_MUL_BLOCK_SIZE 256
136+
134137
#define CUDA_DEQUANTIZE_BLOCK_SIZE 256
135-
#define CUDA_DMMV_BLOCK_SIZE 64 // dmmv = dequantize_mul_mat_vec
138+
139+
// dmmv = dequantize_mul_mat_vec
140+
#ifndef GGML_CUDA_DMMV_X
141+
#define GGML_CUDA_DMMV_X 32
142+
#endif
143+
#ifndef GGML_CUDA_DMMV_Y
144+
#define GGML_CUDA_DMMV_Y 1
145+
#endif
136146

137147
static __global__ void mul_f32(const float * x, const float * y, float * dst, const int kx, const int ky) {
138148
const int i = blockDim.x*blockIdx.x + threadIdx.x;
@@ -247,41 +257,51 @@ static __global__ void dequantize_block(const void * vx, float * y, const int k)
247257
dequantize_kernel(vx, ib, iqs, v0, v1);
248258
}
249259

250-
template <int block_size, int qk, int qr, dequantize_kernel_t dequantize_kernel>
260+
template <int qk, int qr, dequantize_kernel_t dequantize_kernel>
251261
static __global__ void dequantize_mul_mat_vec(const void * vx, const float * y, float * dst, const int ncols) {
252-
const int row = blockIdx.x;
262+
// qk = quantized weights per x block
263+
// qr = number of quantized weights per data value in x block
264+
const int row = blockIdx.x*blockDim.y + threadIdx.y;
253265
const int tid = threadIdx.x;
254266

267+
const int iter_stride = 2*GGML_CUDA_DMMV_X;
268+
const int vals_per_iter = iter_stride / WARP_SIZE; // num quantized vals per thread and i iter
255269
const int y_offset = qr == 1 ? 1 : qk/2;
256270

257-
__shared__ float tmp[block_size]; // separate sum for each thread
258-
tmp[tid] = 0;
271+
float tmp = 0; // partial sum for thread in warp
259272

260-
for (int i = 0; i < ncols/block_size; i += 2) {
261-
const int col = i*block_size + 2*tid;
262-
const int ib = (row*ncols + col)/qk; // block index
263-
const int iqs = (col%qk)/qr; // quant index
273+
for (int i = 0; i < ncols; i += iter_stride) {
274+
const int col = i + vals_per_iter*tid;
275+
const int ib = (row*ncols + col)/qk; // x block index
276+
const int iqs = (col%qk)/qr; // x quant index
264277
const int iybs = col - col%qk; // y block start index
265278

266-
// dequantize
267-
float v0, v1;
268-
dequantize_kernel(vx, ib, iqs, v0, v1);
279+
// processing >2 values per i iter is faster for fast GPUs
280+
#pragma unroll
281+
for (int j = 0; j < vals_per_iter; j += 2) {
282+
// process 2 vals per j iter
283+
284+
// dequantize
285+
float v0, v1;
286+
dequantize_kernel(vx, ib, iqs + j/qr, v0, v1);
287+
// for qr = 2 the iqs needs to increase by 1 per j iter because 2 weights per data val
269288

270-
// matrix multiplication
271-
tmp[tid] += v0 * y[iybs + iqs + 0];
272-
tmp[tid] += v1 * y[iybs + iqs + y_offset];
289+
// matrix multiplication
290+
tmp += v0 * y[iybs + iqs + j/qr + 0];
291+
tmp += v1 * y[iybs + iqs + j/qr + y_offset];
292+
// for qr = 2 the y index needs to increase by 1 per j iter because of y_offset = qk/2
293+
}
273294
}
274295

275296
// sum up partial sums and write back result
276297
__syncthreads();
277-
for (int s=block_size/2; s>0; s>>=1) {
278-
if (tid < s) {
279-
tmp[tid] += tmp[tid + s];
280-
}
281-
__syncthreads();
298+
#pragma unroll
299+
for (int mask = 16; mask > 0; mask >>= 1) {
300+
tmp += __shfl_xor_sync(0xffffffff, tmp, mask, 32);
282301
}
302+
283303
if (tid == 0) {
284-
dst[row] = tmp[0];
304+
dst[row] = tmp;
285305
}
286306
}
287307

@@ -316,33 +336,43 @@ static void dequantize_row_q8_0_cuda(const void * vx, float * y, const int k, cu
316336
}
317337

318338
static void dequantize_mul_mat_vec_q4_0_cuda(const void * vx, const float * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) {
319-
GGML_ASSERT(ncols % CUDA_DMMV_BLOCK_SIZE == 0);
320-
dequantize_mul_mat_vec<CUDA_DMMV_BLOCK_SIZE, QK4_0, QR4_0, dequantize_q4_0>
321-
<<<nrows, CUDA_DMMV_BLOCK_SIZE, 0, stream>>>(vx, y, dst, ncols);
339+
GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0);
340+
GGML_ASSERT(nrows % GGML_CUDA_DMMV_Y == 0);
341+
const dim3 block_dims(WARP_SIZE, GGML_CUDA_DMMV_Y, 1);
342+
dequantize_mul_mat_vec<QK4_0, QR4_0, dequantize_q4_0>
343+
<<<nrows/GGML_CUDA_DMMV_Y, block_dims, 0, stream>>>(vx, y, dst, ncols);
322344
}
323345

324346
static void dequantize_mul_mat_vec_q4_1_cuda(const void * vx, const float * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) {
325-
GGML_ASSERT(ncols % CUDA_DMMV_BLOCK_SIZE == 0);
326-
dequantize_mul_mat_vec<CUDA_DMMV_BLOCK_SIZE, QK4_1, QR4_1, dequantize_q4_1>
327-
<<<nrows, CUDA_DMMV_BLOCK_SIZE, 0, stream>>>(vx, y, dst, ncols);
347+
GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0);
348+
GGML_ASSERT(nrows % GGML_CUDA_DMMV_Y == 0);
349+
const dim3 block_dims(WARP_SIZE, GGML_CUDA_DMMV_Y, 1);
350+
dequantize_mul_mat_vec<QK4_1, QR4_1, dequantize_q4_1>
351+
<<<nrows/GGML_CUDA_DMMV_Y, block_dims, 0, stream>>>(vx, y, dst, ncols);
328352
}
329353

330354
static void dequantize_mul_mat_vec_q5_0_cuda(const void * vx, const float * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) {
331-
GGML_ASSERT(ncols % CUDA_DMMV_BLOCK_SIZE == 0);
332-
dequantize_mul_mat_vec<CUDA_DMMV_BLOCK_SIZE, QK5_0, QR5_0, dequantize_q5_0>
333-
<<<nrows, CUDA_DMMV_BLOCK_SIZE, 0, stream>>>(vx, y, dst, ncols);
355+
GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0);
356+
GGML_ASSERT(nrows % GGML_CUDA_DMMV_Y == 0);
357+
const dim3 block_dims(WARP_SIZE, GGML_CUDA_DMMV_Y, 1);
358+
dequantize_mul_mat_vec<QK5_0, QR5_0, dequantize_q5_0>
359+
<<<nrows/GGML_CUDA_DMMV_Y, block_dims, 0, stream>>>(vx, y, dst, ncols);
334360
}
335361

336362
static void dequantize_mul_mat_vec_q5_1_cuda(const void * vx, const float * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) {
337-
GGML_ASSERT(ncols % CUDA_DMMV_BLOCK_SIZE == 0);
338-
dequantize_mul_mat_vec<CUDA_DMMV_BLOCK_SIZE, QK5_1, QR5_1, dequantize_q5_1>
339-
<<<nrows, CUDA_DMMV_BLOCK_SIZE, 0, stream>>>(vx, y, dst, ncols);
363+
GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0);
364+
GGML_ASSERT(nrows % GGML_CUDA_DMMV_Y == 0);
365+
const dim3 block_dims(WARP_SIZE, GGML_CUDA_DMMV_Y, 1);
366+
dequantize_mul_mat_vec<QK5_1, QR5_1, dequantize_q5_1>
367+
<<<nrows/GGML_CUDA_DMMV_Y, block_dims, 0, stream>>>(vx, y, dst, ncols);
340368
}
341369

342370
static void dequantize_mul_mat_vec_q8_0_cuda(const void * vx, const float * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) {
343-
GGML_ASSERT(ncols % CUDA_DMMV_BLOCK_SIZE == 0);
344-
dequantize_mul_mat_vec<CUDA_DMMV_BLOCK_SIZE, QK8_0, QR8_0, dequantize_q8_0>
345-
<<<nrows, CUDA_DMMV_BLOCK_SIZE, 0, stream>>>(vx, y, dst, ncols);
371+
GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0);
372+
GGML_ASSERT(nrows % GGML_CUDA_DMMV_Y == 0);
373+
const dim3 block_dims(WARP_SIZE, GGML_CUDA_DMMV_Y, 1);
374+
dequantize_mul_mat_vec<QK8_0, QR8_0, dequantize_q8_0>
375+
<<<nrows/GGML_CUDA_DMMV_Y, block_dims, 0, stream>>>(vx, y, dst, ncols);
346376
}
347377

348378
static void convert_fp16_to_fp32_cuda(const void * vx, float * y, const int k, cudaStream_t stream) {
@@ -351,9 +381,11 @@ static void convert_fp16_to_fp32_cuda(const void * vx, float * y, const int k, c
351381
}
352382

353383
static void convert_mul_mat_vec_f16_cuda(const void * vx, const float * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) {
354-
GGML_ASSERT(ncols % CUDA_DMMV_BLOCK_SIZE == 0);
355-
dequantize_mul_mat_vec<CUDA_DMMV_BLOCK_SIZE, 32, 1, convert_f16>
356-
<<<nrows, CUDA_DMMV_BLOCK_SIZE, 0, stream>>>(vx, y, dst, ncols);
384+
GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0);
385+
GGML_ASSERT(nrows % GGML_CUDA_DMMV_Y == 0);
386+
const dim3 block_dims(WARP_SIZE, GGML_CUDA_DMMV_Y, 1);
387+
dequantize_mul_mat_vec<1, 1, convert_f16>
388+
<<<nrows/GGML_CUDA_DMMV_Y, block_dims, 0, stream>>>(vx, y, dst, ncols);
357389
}
358390

359391
static to_fp32_cuda_t ggml_get_to_fp32_cuda(ggml_type type) {

0 commit comments

Comments
 (0)