forked from mrq/bitsandbytes-rocm
Merge pull request #3 from TimDettmers/cpuonly
Add a CPU-only build option
This commit is contained in:
commit
4cd7ea62b2
19
Makefile
19
Makefile
|
@ -10,10 +10,10 @@ NVCC := $(CUDA_HOME)/bin/nvcc
|
|||
###########################################
|
||||
|
||||
CSRC := $(ROOT_DIR)/csrc
|
||||
BUILD_DIR:= $(ROOT_DIR)/cuda_build
|
||||
BUILD_DIR:= $(ROOT_DIR)/build
|
||||
|
||||
FILES_CUDA := $(CSRC)/ops.cu $(CSRC)/kernels.cu
|
||||
FILES_CPP := $(CSRC)/pythonInterface.c
|
||||
FILES_CPP := $(CSRC)/common.cpp $(CSRC)/cpu_ops.cpp $(CSRC)/pythonInterface.c
|
||||
|
||||
INCLUDE := -I $(CUDA_HOME)/include -I $(ROOT_DIR)/csrc -I $(CONDA_PREFIX)/include -I $(ROOT_DIR)/dependencies/cub -I $(ROOT_DIR)/include
|
||||
LIB := -L $(CUDA_HOME)/lib64 -lcudart -lcuda -lcublas -lcurand -lcusparse -L $(CONDA_PREFIX)/lib
|
||||
|
@ -46,27 +46,30 @@ CC_CUDA11x += -gencode arch=compute_86,code=sm_86
|
|||
all: $(ROOT_DIR)/dependencies/cub $(BUILD_DIR) env
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) -Xcompiler '-fPIC' --use_fast_math -Xptxas=-v -dc $(FILES_CUDA) $(INCLUDE) $(LIB) --output-directory $(BUILD_DIR)
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) -Xcompiler '-fPIC' -dlink $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o -o $(BUILD_DIR)/link.o
|
||||
$(GPP) -std=c++14 -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
$(GPP) -std=c++14 -DBUILD_CUDA -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
|
||||
cuda92: $(ROOT_DIR)/dependencies/cub $(BUILD_DIR) env
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA92) -Xcompiler '-fPIC' --use_fast_math -Xptxas=-v -dc $(FILES_CUDA) $(INCLUDE) $(LIB) --output-directory $(BUILD_DIR)
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA92) -Xcompiler '-fPIC' -dlink $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o -o $(BUILD_DIR)/link.o
|
||||
$(GPP) -std=c++14 -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
$(GPP) -std=c++14 -DBUILD_CUDA -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
|
||||
cuda10x: $(ROOT_DIR)/dependencies/cub $(BUILD_DIR) env
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA10x) -Xcompiler '-fPIC' --use_fast_math -Xptxas=-v -dc $(FILES_CUDA) $(INCLUDE) $(LIB) --output-directory $(BUILD_DIR)
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA10x) -Xcompiler '-fPIC' -dlink $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o -o $(BUILD_DIR)/link.o
|
||||
$(GPP) -std=c++14 -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
$(GPP) -std=c++14 -DBUILD_CUDA -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
|
||||
cuda110: $(BUILD_DIR) env
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA110) -Xcompiler '-fPIC' --use_fast_math -Xptxas=-v -dc $(FILES_CUDA) $(INCLUDE) $(LIB) --output-directory $(BUILD_DIR)
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA110) -Xcompiler '-fPIC' -dlink $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o -o $(BUILD_DIR)/link.o
|
||||
$(GPP) -std=c++14 -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
$(GPP) -std=c++14 -DBUILD_CUDA -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
|
||||
cuda11x: $(BUILD_DIR) env
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA11x) -Xcompiler '-fPIC' --use_fast_math -Xptxas=-v -dc $(FILES_CUDA) $(INCLUDE) $(LIB) --output-directory $(BUILD_DIR)
|
||||
$(NVCC) $(COMPUTE_CAPABILITY) $(CC_CUDA11x) -Xcompiler '-fPIC' -dlink $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o -o $(BUILD_DIR)/link.o
|
||||
$(GPP) -std=c++14 -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
$(GPP) -std=c++14 -DBUILD_CUDA -shared -fPIC $(INCLUDE) $(BUILD_DIR)/ops.o $(BUILD_DIR)/kernels.o $(BUILD_DIR)/link.o $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so $(LIB)
|
||||
|
||||
cpuonly: $(BUILD_DIR) env
|
||||
$(GPP) -std=c++14 -shared -fPIC -I $(ROOT_DIR)/csrc -I $(ROOT_DIR)/include $(FILES_CPP) -o ./bitsandbytes/libbitsandbytes.so
|
||||
|
||||
env:
|
||||
@echo "ENVIRONMENT"
|
||||
|
@ -80,7 +83,7 @@ env:
|
|||
@echo "============================"
|
||||
|
||||
$(BUILD_DIR):
|
||||
mkdir -p cuda_build
|
||||
mkdir -p build
|
||||
mkdir -p dependencies
|
||||
|
||||
$(ROOT_DIR)/dependencies/cub:
|
||||
|
|
|
@ -2,9 +2,14 @@
|
|||
#
|
||||
# This source code is licensed under the MIT license found in the
|
||||
# LICENSE file in the root directory of this source tree.
|
||||
from .optim import adam
|
||||
|
||||
from .nn import modules
|
||||
__pdoc__ = {'libBitsNBytes' : False,
|
||||
from .cextension import COMPILED_WITH_CUDA
|
||||
|
||||
if COMPILED_WITH_CUDA:
|
||||
from .optim import adam
|
||||
|
||||
__pdoc__ = {'libBitsNBytes': False,
|
||||
'optim.optimizer.Optimizer8bit': False,
|
||||
'optim.optimizer.MockArgs': False
|
||||
}
|
||||
|
|
13
bitsandbytes/cextension.py
Normal file
13
bitsandbytes/cextension.py
Normal file
|
@ -0,0 +1,13 @@
|
|||
import ctypes as ct
|
||||
import os
|
||||
from warnings import warn
|
||||
|
||||
lib = ct.cdll.LoadLibrary(os.path.dirname(__file__) + '/libbitsandbytes.so')
|
||||
|
||||
try:
|
||||
lib.cadam32bit_g32
|
||||
COMPILED_WITH_CUDA = True
|
||||
except AttributeError:
|
||||
warn("The installed version of bitsandbytes was compiled without GPU support. "
|
||||
"8-bit optimizers and GPU quantization are unavailable.")
|
||||
COMPILED_WITH_CUDA = False
|
File diff suppressed because one or more lines are too long
|
@ -2,11 +2,16 @@
|
|||
#
|
||||
# This source code is licensed under the MIT license found in the
|
||||
# LICENSE file in the root directory of this source tree.
|
||||
from .adam import Adam, Adam8bit, Adam32bit
|
||||
from .adamw import AdamW, AdamW8bit, AdamW32bit
|
||||
from .sgd import SGD, SGD8bit, SGD32bit
|
||||
from .lars import LARS, LARS8bit, LARS32bit, PytorchLARS
|
||||
from .lamb import LAMB, LAMB8bit, LAMB32bit
|
||||
from .rmsprop import RMSprop, RMSprop8bit, RMSprop32bit
|
||||
from .adagrad import Adagrad, Adagrad8bit, Adagrad32bit
|
||||
|
||||
from bitsandbytes.cextension import COMPILED_WITH_CUDA
|
||||
|
||||
if COMPILED_WITH_CUDA:
|
||||
from .adam import Adam, Adam8bit, Adam32bit
|
||||
from .adamw import AdamW, AdamW8bit, AdamW32bit
|
||||
from .sgd import SGD, SGD8bit, SGD32bit
|
||||
from .lars import LARS, LARS8bit, LARS32bit, PytorchLARS
|
||||
from .lamb import LAMB, LAMB8bit, LAMB32bit
|
||||
from .rmsprop import RMSprop, RMSprop8bit, RMSprop32bit
|
||||
from .adagrad import Adagrad, Adagrad8bit, Adagrad32bit
|
||||
|
||||
from .optimizer import GlobalOptimManager
|
||||
|
|
|
@ -31,6 +31,6 @@ class RMSprop32bit(Optimizer1State):
|
|||
if alpha == 0:
|
||||
raise NotImplementedError(f'RMSprop with alpha==0.0 is not supported!')
|
||||
if centered:
|
||||
raise NotImplementError(f'Centered RMSprop is not supported!')
|
||||
raise NotImplementedError(f'Centered RMSprop is not supported!')
|
||||
super(RMSprop32bit, self).__init__('rmsprop', params, lr, (alpha, momentum), eps,
|
||||
weight_decay, 32, args, min_8bit_size, percentile_clipping, block_wise)
|
||||
|
|
|
@ -1,12 +1,12 @@
|
|||
# Compiling from source
|
||||
|
||||
Basic steps.
|
||||
1. `make cudaXXX` where `cudaXXX` is among `cuda92, cuda10x, cuda110, cuda11x`
|
||||
1. `make [target]` where `[target]` is among `cuda92, cuda10x, cuda110, cuda11x, cpuonly`
|
||||
2. `CUDA_VERSION=XXX python setup.py install`
|
||||
|
||||
To run these steps you will need to have the nvcc compiler installed that comes with a CUDA installation. If you use anaconda (recommended) then you can figure out which version of CUDA you are using with PyTorch via the command `conda list | grep cudatoolkit`. Then you can install the nvcc compiler by downloading and installing the same CUDA version from the [CUDA toolkit archive](https://developer.nvidia.com/cuda-toolkit-archive).
|
||||
|
||||
For your convenience, there is a install script int he root directory that installs CUDA 11.1 locally and configures it automatically. After installing you should add the `bin` sub-directory to the `$PATH` variable to make the compiler visible to your system. To do this you can add this to your `.bashrc` by executing these commands:
|
||||
For your convenience, there is an installation script in the root directory that installs CUDA 11.1 locally and configures it automatically. After installing you should add the `bin` sub-directory to the `$PATH` variable to make the compiler visible to your system. To do this you can add this to your `.bashrc` by executing these commands:
|
||||
```bash
|
||||
echo "export LD_LIBRARY_PATH=$LD_LIBRARY_PATH:/usr/local/cuda/lib64/" >> ~/.bashrc
|
||||
echo "export PATH=$PATH:/usr/local/cuda/bin/" >> ~/.bashrc
|
||||
|
|
39
csrc/common.cpp
Normal file
39
csrc/common.cpp
Normal file
|
@ -0,0 +1,39 @@
|
|||
#include <common.h>
|
||||
#include <float.h>
|
||||
|
||||
void *quantize_block(void *arguments) {
|
||||
// 1. find absmax in block
|
||||
// 2. divide input value by absmax to normalize into [-1.0, 1.0]
|
||||
// 3. do binary search to find the closest value
|
||||
// 4. check minimal distance
|
||||
// 5. store index
|
||||
|
||||
struct quantize_block_args *args = (quantize_block_args *) arguments;
|
||||
|
||||
// 1. find absmax in block
|
||||
float absmax_block = -FLT_MAX;
|
||||
for (int i = args->block_idx; i < args->block_end; i++)
|
||||
absmax_block = fmax(absmax_block, fabs(args->A[i]));
|
||||
|
||||
args->absmax[args->block_idx / BLOCK_SIZE] = absmax_block;
|
||||
|
||||
for (int i = args->block_idx; i < args->block_end; i++) {
|
||||
// 2. divide input value by absmax to normalize into [-1.0, 1.0]
|
||||
// 3. do binary search to find the closest value
|
||||
float normed_value = args->A[i] / absmax_block;
|
||||
int idx = args->bin_searcher->scalar(normed_value);
|
||||
|
||||
// 4. check minimal distance
|
||||
// The binary search returns always the value to the left, which might not be the closest value
|
||||
if (idx < 255) {
|
||||
float dist_left = fabs(normed_value - (args->code[idx]));
|
||||
float dist_right = fabs(normed_value - (args->code[idx + 1]));
|
||||
if (dist_right < dist_left) { idx += 1; }
|
||||
}
|
||||
|
||||
// 5. store index
|
||||
args->out[i] = (unsigned char) idx;
|
||||
}
|
||||
|
||||
return NULL;
|
||||
}
|
23
csrc/common.h
Normal file
23
csrc/common.h
Normal file
|
@ -0,0 +1,23 @@
|
|||
#include <BinSearch.h>
|
||||
|
||||
#ifndef common
|
||||
#define common
|
||||
|
||||
using namespace BinSearch;
|
||||
|
||||
struct quantize_block_args {
|
||||
BinAlgo<Scalar, float, Direct2> *bin_searcher;
|
||||
float *code;
|
||||
float *A;
|
||||
float *absmax;
|
||||
unsigned char *out;
|
||||
int block_end;
|
||||
int block_idx;
|
||||
int threadidx;
|
||||
};
|
||||
|
||||
#define BLOCK_SIZE 4096
|
||||
|
||||
void *quantize_block(void *arguments);
|
||||
|
||||
#endif
|
57
csrc/cpu_ops.cpp
Normal file
57
csrc/cpu_ops.cpp
Normal file
|
@ -0,0 +1,57 @@
|
|||
#include <BinSearch.h>
|
||||
#include <pthread.h>
|
||||
#include <common.h>
|
||||
|
||||
using namespace BinSearch;
|
||||
|
||||
void dequantize_cpu(float *code, unsigned char *A, float *absmax, float *out, int n) {
|
||||
for (int block_idx = 0; block_idx < n; block_idx += BLOCK_SIZE) {
|
||||
int valid_items = n - block_idx >= BLOCK_SIZE ? BLOCK_SIZE : n - block_idx;
|
||||
int block_end = block_idx + valid_items;
|
||||
for (int i = block_idx; i < block_end; i++)
|
||||
out[i] = code[A[i]] * absmax[block_idx / BLOCK_SIZE];
|
||||
}
|
||||
}
|
||||
|
||||
void quantize_cpu(float *code, float *A, float *absmax, unsigned char *out, int n) {
|
||||
|
||||
// the default code is has range [-0.993, 1.0] which can cause an error in the binary search algorithm used below
|
||||
code[0] = -1.0f;
|
||||
|
||||
int num_blocks = n / BLOCK_SIZE;
|
||||
num_blocks += n % BLOCK_SIZE == 0 ? 0 : 1;
|
||||
|
||||
pthread_t *threads = (pthread_t *) malloc(sizeof(pthread_t) * num_blocks);
|
||||
struct quantize_block_args **args = (quantize_block_args **) malloc(num_blocks * sizeof(quantize_block_args *));
|
||||
|
||||
for (int i = 0; i < num_blocks; i++)
|
||||
args[i] = (quantize_block_args *) malloc(sizeof(quantize_block_args));
|
||||
|
||||
const uint32 elements_code = 256;
|
||||
BinAlgo<Scalar, float, Direct2> bin_searcher(code, elements_code);
|
||||
|
||||
for (int block_idx = 0; block_idx < n; block_idx += BLOCK_SIZE) {
|
||||
int valid_items = n - block_idx >= BLOCK_SIZE ? BLOCK_SIZE : n - block_idx;
|
||||
int block_end = block_idx + valid_items;
|
||||
|
||||
struct quantize_block_args *arg = args[block_idx / BLOCK_SIZE];
|
||||
arg->bin_searcher = &bin_searcher;
|
||||
arg->code = code;
|
||||
arg->A = A;
|
||||
arg->absmax = absmax;
|
||||
arg->out = out;
|
||||
arg->block_end = block_end;
|
||||
arg->block_idx = block_idx;
|
||||
arg->threadidx = block_idx / BLOCK_SIZE;
|
||||
|
||||
pthread_create(&threads[block_idx / BLOCK_SIZE], NULL, &quantize_block, (void *) arg);
|
||||
}
|
||||
|
||||
for (int i = 0; i < num_blocks; i++)
|
||||
int err = pthread_join(threads[i], NULL);
|
||||
|
||||
free(threads);
|
||||
for (int i = 0; i < num_blocks; i++)
|
||||
free(args[i]);
|
||||
free(args);
|
||||
}
|
9
csrc/cpu_ops.h
Normal file
9
csrc/cpu_ops.h
Normal file
|
@ -0,0 +1,9 @@
|
|||
#ifndef BITSANDBYTES_CPU_OPS_H
|
||||
#define BITSANDBYTES_CPU_OPS_H
|
||||
|
||||
|
||||
void quantize_cpu(float *code, float *A, float *absmax, unsigned char *out, int n);
|
||||
|
||||
void dequantize_cpu(float *code, unsigned char *A, float *absmax, float *out, int n);
|
||||
|
||||
#endif
|
112
csrc/ops.cu
112
csrc/ops.cu
|
@ -8,123 +8,13 @@
|
|||
#include <cub/device/device_scan.cuh>
|
||||
#include <limits>
|
||||
#include <BinSearch.h>
|
||||
#include <common.h>
|
||||
|
||||
|
||||
using namespace BinSearch;
|
||||
using std::cout;
|
||||
using std::endl;
|
||||
|
||||
#define BLOCK_SIZE 4096
|
||||
|
||||
struct quantize_block_args
|
||||
{
|
||||
BinAlgo<Scalar, float, Direct2> *bin_searcher;
|
||||
float *code;
|
||||
float *A;
|
||||
float *absmax;
|
||||
unsigned char *out;
|
||||
int block_end;
|
||||
int block_idx;
|
||||
int threadidx;
|
||||
};
|
||||
|
||||
void *quantize_block(void *arguments)
|
||||
{
|
||||
// 1. find absmax in block
|
||||
// 2. divide input value by absmax to normalize into [-1.0, 1.0]
|
||||
// 3. do binary search to find the closest value
|
||||
// 4. check minimal distance
|
||||
// 5. store index
|
||||
|
||||
struct quantize_block_args *args = (quantize_block_args*)arguments;
|
||||
|
||||
// 1. find absmax in block
|
||||
float absmax_block = -FLT_MAX;
|
||||
for (int i = args->block_idx; i < args->block_end; i++)
|
||||
absmax_block = fmax(absmax_block, fabs(args->A[i]));
|
||||
|
||||
args->absmax[args->block_idx/BLOCK_SIZE] = absmax_block;
|
||||
|
||||
for (int i = args->block_idx; i < args->block_end; i++)
|
||||
{
|
||||
// 2. divide input value by absmax to normalize into [-1.0, 1.0]
|
||||
// 3. do binary search to find the closest value
|
||||
float normed_value = args->A[i]/absmax_block;
|
||||
int idx = args->bin_searcher->scalar(normed_value);
|
||||
|
||||
// 4. check minimal distance
|
||||
// The binary search returns always the value to the left, which might not be the closest value
|
||||
if(idx < 255)
|
||||
{
|
||||
float dist_left = fabs(normed_value-(args->code[idx]));
|
||||
float dist_right = fabs(normed_value-(args->code[idx+1]));
|
||||
if(dist_right < dist_left){ idx+=1; }
|
||||
}
|
||||
|
||||
// 5. store index
|
||||
args->out[i] = (unsigned char)idx;
|
||||
}
|
||||
|
||||
return NULL;
|
||||
}
|
||||
|
||||
void quantize_cpu(float *code, float *A, float *absmax, unsigned char *out, int n)
|
||||
{
|
||||
|
||||
// the default code is has range [-0.993, 1.0] which can cause an error in the binary search algorithm used below
|
||||
code[0] = -1.0f;
|
||||
|
||||
int num_blocks = n/BLOCK_SIZE;
|
||||
num_blocks += n % BLOCK_SIZE == 0 ? 0 : 1;
|
||||
|
||||
pthread_t *threads = (pthread_t*)malloc(sizeof(pthread_t)*num_blocks);
|
||||
struct quantize_block_args **args = (quantize_block_args**)malloc(num_blocks*sizeof(quantize_block_args*));
|
||||
|
||||
for(int i = 0; i < num_blocks; i++)
|
||||
args[i] = (quantize_block_args*)malloc(sizeof(quantize_block_args));
|
||||
|
||||
const uint32 elements_code = 256;
|
||||
BinAlgo<Scalar, float, Direct2> bin_searcher(code, elements_code);
|
||||
|
||||
for(int block_idx = 0; block_idx < n; block_idx+=BLOCK_SIZE)
|
||||
{
|
||||
int valid_items = n-block_idx >= BLOCK_SIZE ? BLOCK_SIZE : n - block_idx;
|
||||
int block_end = block_idx + valid_items;
|
||||
|
||||
struct quantize_block_args *arg = args[block_idx/BLOCK_SIZE];
|
||||
arg->bin_searcher = &bin_searcher;
|
||||
arg->code = code;
|
||||
arg->A = A;
|
||||
arg->absmax = absmax;
|
||||
arg->out = out;
|
||||
arg->block_end = block_end;
|
||||
arg->block_idx = block_idx;
|
||||
arg->threadidx = block_idx/BLOCK_SIZE;
|
||||
|
||||
pthread_create(&threads[block_idx/BLOCK_SIZE], NULL, &quantize_block, (void *)arg);
|
||||
}
|
||||
|
||||
for(int i = 0; i < num_blocks; i++)
|
||||
int err = pthread_join(threads[i], NULL);
|
||||
|
||||
free(threads);
|
||||
for(int i = 0; i < num_blocks; i++)
|
||||
free(args[i]);
|
||||
free(args);
|
||||
}
|
||||
|
||||
|
||||
void dequantize_cpu(float *code, unsigned char *A, float *absmax, float *out, int n)
|
||||
{
|
||||
for(int block_idx = 0; block_idx < n; block_idx+=BLOCK_SIZE)
|
||||
{
|
||||
int valid_items = n-block_idx >= BLOCK_SIZE ? BLOCK_SIZE : n - block_idx;
|
||||
int block_end = block_idx + valid_items;
|
||||
for (int i = block_idx; i < block_end; i++)
|
||||
out[i] = code[A[i]]*absmax[block_idx/BLOCK_SIZE];
|
||||
}
|
||||
}
|
||||
|
||||
void histogramScatterAdd2D(float* histogram, int *index1, int *index2, float *src, int maxidx1, int n)
|
||||
{
|
||||
int threads = 512;
|
||||
|
|
10
csrc/ops.cuh
10
csrc/ops.cuh
|
@ -68,16 +68,6 @@ template<typename T, int OPTIMIZER> void optimizerStatic8bitBlockwise(T* p, T* g
|
|||
|
||||
template<typename T> void percentileClipping(T * g, float *gnorm_vec, int step, const int n);
|
||||
|
||||
void quantize_cpu(float *code, float *A, float *absmax, unsigned char *out, int n);
|
||||
void dequantize_cpu(float *code, unsigned char *A, float *absmax, float *out, int n);
|
||||
|
||||
void histogramScatterAdd2D(float* histogram, int *index1, int *index2, float *src, int maxidx1, int n);
|
||||
|
||||
#endif
|
||||
|
||||
|
||||
|
||||
|
||||
|
||||
|
||||
|
||||
|
|
|
@ -3,7 +3,10 @@
|
|||
// This source code is licensed under the MIT license found in the
|
||||
// LICENSE file in the root directory of this source tree.
|
||||
|
||||
#if BUILD_CUDA
|
||||
#include <ops.cuh>
|
||||
#endif
|
||||
#include <cpu_ops.h>
|
||||
|
||||
// We cannot call templated code from C, so we wrap the template in a C compatible call here if necessary.
|
||||
// We use macro functions to expand all the different optimizers. Looks ugly, and is ugly, but its better than to
|
||||
|
@ -12,6 +15,7 @@
|
|||
// UNMANGLED CALLS
|
||||
//===================================================================================
|
||||
|
||||
#if BUILD_CUDA
|
||||
void estimateQuantiles_fp32(float *A, float *code, float offset, int n){ estimateQuantiles<float>(A, code, offset, n); }
|
||||
void estimateQuantiles_fp16(half *A, float *code, float offset, int n){ estimateQuantiles<half>(A, code, offset, n); }
|
||||
|
||||
|
@ -78,9 +82,11 @@ void quantizeBlockwise_stochastic_fp32(float * code, float *A, float *absmax, un
|
|||
|
||||
void dequantizeBlockwise_fp16(float *code, unsigned char *A, float *absmax, half *out, int blocksize, const int n){ dequantizeBlockwise<half>(code, A, absmax, out, blocksize, n); } \
|
||||
void dequantizeBlockwise_fp32(float *code, unsigned char *A, float *absmax, float *out, int blocksize, const int n){ dequantizeBlockwise<float>(code, A, absmax, out, blocksize, n); }
|
||||
#endif
|
||||
|
||||
extern "C"
|
||||
{
|
||||
#if BUILD_CUDA
|
||||
void cestimate_quantiles_fp32(float *A, float *code, float offset, int n){ estimateQuantiles_fp32(A, code, offset, n); }
|
||||
void cestimate_quantiles_fp16(half *A, float *code, float offset, int n){ estimateQuantiles_fp16(A, code, offset, n); }
|
||||
void cquantize(float *code, float *A, unsigned char *out, int n){ quantize(code, A, out, n); }
|
||||
|
@ -147,11 +153,10 @@ extern "C"
|
|||
|
||||
void cpercentile_clipping_g32(float * g, float *gnorm_vec, int step, const int n){ percentileClipping_g32(g, gnorm_vec, step, n); }
|
||||
void cpercentile_clipping_g16(half * g, float *gnorm_vec, int step, const int n){ percentileClipping_g16(g, gnorm_vec, step, n); }
|
||||
void chistogram_scatter_add_2d(float* histogram, int *index1, int *index2, float *src, int maxidx1, int n){ histogramScatterAdd2D(histogram, index1, index2, src, maxidx1, n); }
|
||||
|
||||
#endif
|
||||
void cquantize_blockwise_cpu_fp32(float *code, float *A, float *absmax, unsigned char *out, const int n){ quantize_cpu(code, A, absmax, out, n); }
|
||||
void cdequantize_blockwise_cpu_fp32(float *code, unsigned char *A, float *absmax, float *out, const int n){ dequantize_cpu(code, A, absmax, out, n); }
|
||||
|
||||
void chistogram_scatter_add_2d(float* histogram, int *index1, int *index2, float *src, int maxidx1, int n){ histogramScatterAdd2D(histogram, index1, index2, src, maxidx1, n); }
|
||||
}
|
||||
|
||||
|
||||
|
|
22
setup.py
22
setup.py
|
@ -6,27 +6,27 @@ import os
|
|||
from setuptools import setup, find_packages
|
||||
|
||||
|
||||
|
||||
def read(fname):
|
||||
return open(os.path.join(os.path.dirname(__file__), fname)).read()
|
||||
|
||||
|
||||
version = os.getenv("CUDA_VERSION", "cpu")
|
||||
|
||||
setup(
|
||||
name = f"bitsandbytes-cuda{os.environ['CUDA_VERSION']}",
|
||||
version = "0.26.0",
|
||||
author = "Tim Dettmers",
|
||||
author_email = "dettmers@cs.washington.edu",
|
||||
description = ("8-bit optimizers and quantization routines."),
|
||||
license = "MIT",
|
||||
keywords = "gpu optimizers optimization 8-bit quantization compression",
|
||||
url = "http://packages.python.org/bitsandbytes",
|
||||
name="bitsandbytes",
|
||||
version=f"0.26.0+{version}",
|
||||
author="Tim Dettmers",
|
||||
author_email="dettmers@cs.washington.edu",
|
||||
description="8-bit optimizers and quantization routines.",
|
||||
license="MIT",
|
||||
keywords="gpu optimizers optimization 8-bit quantization compression",
|
||||
url="http://packages.python.org/bitsandbytes",
|
||||
packages=find_packages(),
|
||||
package_data={'': ['libbitsandbytes.so']},
|
||||
long_description=read('README.md'),
|
||||
long_description_content_type = 'text/markdown',
|
||||
long_description_content_type='text/markdown',
|
||||
classifiers=[
|
||||
"Development Status :: 4 - Beta",
|
||||
'Topic :: Scientific/Engineering :: Artificial Intelligence'
|
||||
],
|
||||
)
|
||||
|
||||
|
|
Loading…
Reference in New Issue
Block a user